trainer.py 25 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629
  1. # Copyright (c) 2020 PaddlePaddle Authors. All Rights Reserved.
  2. #
  3. # Licensed under the Apache License, Version 2.0 (the "License");
  4. # you may not use this file except in compliance with the License.
  5. # You may obtain a copy of the License at
  6. #
  7. # http://www.apache.org/licenses/LICENSE-2.0
  8. #
  9. # Unless required by applicable law or agreed to in writing, software
  10. # distributed under the License is distributed on an "AS IS" BASIS,
  11. # WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
  12. # See the License for the specific language governing permissions and
  13. # limitations under the License.
  14. from __future__ import absolute_import
  15. from __future__ import division
  16. from __future__ import print_function
  17. import os
  18. import sys
  19. import copy
  20. import time
  21. import numpy as np
  22. from PIL import Image
  23. import paddle
  24. import paddle.distributed as dist
  25. from paddle.distributed import fleet
  26. from paddle import amp
  27. from paddle.static import InputSpec
  28. from paddlex.ppdet.optimizer import ModelEMA
  29. from paddlex.ppdet.core.workspace import create
  30. from paddlex.ppdet.utils.checkpoint import load_weight, load_pretrain_weight
  31. from paddlex.ppdet.utils.visualizer import visualize_results, save_result
  32. from paddlex.ppdet.metrics import Metric, COCOMetric, VOCMetric, WiderFaceMetric, get_infer_results, KeyPointTopDownCOCOEval, KeyPointTopDownMPIIEval
  33. from paddlex.ppdet.metrics import RBoxMetric, JDEDetMetric
  34. from paddlex.ppdet.data.source.category import get_categories
  35. from paddlex.ppdet.utils import stats
  36. from .callbacks import Callback, ComposeCallback, LogPrinter, Checkpointer, WiferFaceEval, VisualDLWriter
  37. from .export_utils import _dump_infer_config
  38. from paddlex.ppdet.utils.logger import setup_logger
  39. logger = setup_logger('paddlex.ppdet.engine')
  40. __all__ = ['Trainer']
  41. MOT_ARCH = ['DeepSORT', 'JDE', 'FairMOT']
  42. class Trainer(object):
  43. def __init__(self, cfg, mode='train'):
  44. self.cfg = cfg
  45. assert mode.lower() in ['train', 'eval', 'test'], \
  46. "mode should be 'train', 'eval' or 'test'"
  47. self.mode = mode.lower()
  48. self.optimizer = None
  49. self.is_loaded_weights = False
  50. # build data loader
  51. if cfg.architecture in MOT_ARCH and self.mode in ['eval', 'test']:
  52. self.dataset = cfg['{}MOTDataset'.format(self.mode.capitalize())]
  53. else:
  54. self.dataset = cfg['{}Dataset'.format(self.mode.capitalize())]
  55. if cfg.architecture == 'DeepSORT' and self.mode == 'train':
  56. logger.error('DeepSORT has no need of training on mot dataset.')
  57. sys.exit(1)
  58. if self.mode == 'train':
  59. self.loader = create('{}Reader'.format(self.mode.capitalize()))(
  60. self.dataset, cfg.worker_num)
  61. if cfg.architecture == 'JDE' and self.mode == 'train':
  62. cfg['JDEEmbeddingHead'][
  63. 'num_identifiers'] = self.dataset.total_identities
  64. if cfg.architecture == 'FairMOT' and self.mode == 'train':
  65. cfg['FairMOTEmbeddingHead'][
  66. 'num_identifiers'] = self.dataset.total_identities
  67. # build model
  68. if 'model' not in self.cfg:
  69. self.model = create(cfg.architecture)
  70. else:
  71. self.model = self.cfg.model
  72. self.is_loaded_weights = True
  73. self.use_ema = ('use_ema' in cfg and cfg['use_ema'])
  74. if self.use_ema:
  75. self.ema = ModelEMA(
  76. cfg['ema_decay'], self.model, use_thres_step=True)
  77. # EvalDataset build with BatchSampler to evaluate in single device
  78. # TODO: multi-device evaluate
  79. if self.mode == 'eval':
  80. self._eval_batch_sampler = paddle.io.BatchSampler(
  81. self.dataset, batch_size=self.cfg.EvalReader['batch_size'])
  82. self.loader = create('{}Reader'.format(self.mode.capitalize()))(
  83. self.dataset, cfg.worker_num, self._eval_batch_sampler)
  84. # TestDataset build after user set images, skip loader creation here
  85. # build optimizer in train mode
  86. if self.mode == 'train':
  87. steps_per_epoch = len(self.loader)
  88. self.lr = create('LearningRate')(steps_per_epoch)
  89. self.optimizer = create('OptimizerBuilder')(
  90. self.lr, self.model.parameters())
  91. self._nranks = dist.get_world_size()
  92. self._local_rank = dist.get_rank()
  93. self.status = {}
  94. self.start_epoch = 0
  95. self.end_epoch = 0 if 'epoch' not in cfg else cfg.epoch
  96. # initial default callbacks
  97. self._init_callbacks()
  98. # initial default metrics
  99. self._init_metrics()
  100. self._reset_metrics()
  101. def _init_callbacks(self):
  102. if self.mode == 'train':
  103. self._callbacks = [LogPrinter(self), Checkpointer(self)]
  104. if self.cfg.get('use_vdl', False):
  105. self._callbacks.append(VisualDLWriter(self))
  106. self._compose_callback = ComposeCallback(self._callbacks)
  107. elif self.mode == 'eval':
  108. self._callbacks = [LogPrinter(self)]
  109. if self.cfg.metric == 'WiderFace':
  110. self._callbacks.append(WiferFaceEval(self))
  111. self._compose_callback = ComposeCallback(self._callbacks)
  112. elif self.mode == 'test' and self.cfg.get('use_vdl', False):
  113. self._callbacks = [VisualDLWriter(self)]
  114. self._compose_callback = ComposeCallback(self._callbacks)
  115. else:
  116. self._callbacks = []
  117. self._compose_callback = None
  118. def _init_metrics(self, validate=False):
  119. if self.mode == 'test' or (self.mode == 'train' and not validate):
  120. self._metrics = []
  121. return
  122. classwise = self.cfg['classwise'] if 'classwise' in self.cfg else False
  123. if self.cfg.metric == 'COCO':
  124. # TODO: bias should be unified
  125. bias = self.cfg['bias'] if 'bias' in self.cfg else 0
  126. output_eval = self.cfg['output_eval'] \
  127. if 'output_eval' in self.cfg else None
  128. save_prediction_only = self.cfg.get('save_prediction_only', False)
  129. # pass clsid2catid info to metric instance to avoid multiple loading
  130. # annotation file
  131. clsid2catid = {v: k for k, v in self.dataset.catid2clsid.items()} \
  132. if self.mode == 'eval' else None
  133. # when do validation in train, annotation file should be get from
  134. # EvalReader instead of self.dataset(which is TrainReader)
  135. anno_file = self.dataset.get_anno()
  136. if self.mode == 'train' and validate:
  137. eval_dataset = self.cfg['EvalDataset']
  138. eval_dataset.check_or_download_dataset()
  139. anno_file = eval_dataset.get_anno()
  140. IouType = self.cfg['IouType'] if 'IouType' in self.cfg else 'bbox'
  141. self._metrics = [
  142. COCOMetric(
  143. anno_file=anno_file,
  144. clsid2catid=clsid2catid,
  145. classwise=classwise,
  146. output_eval=output_eval,
  147. bias=bias,
  148. IouType=IouType,
  149. save_prediction_only=save_prediction_only)
  150. ]
  151. elif self.cfg.metric == 'RBOX':
  152. # TODO: bias should be unified
  153. bias = self.cfg['bias'] if 'bias' in self.cfg else 0
  154. output_eval = self.cfg['output_eval'] \
  155. if 'output_eval' in self.cfg else None
  156. save_prediction_only = self.cfg.get('save_prediction_only', False)
  157. # pass clsid2catid info to metric instance to avoid multiple loading
  158. # annotation file
  159. clsid2catid = {v: k for k, v in self.dataset.catid2clsid.items()} \
  160. if self.mode == 'eval' else None
  161. # when do validation in train, annotation file should be get from
  162. # EvalReader instead of self.dataset(which is TrainReader)
  163. anno_file = self.dataset.get_anno()
  164. if self.mode == 'train' and validate:
  165. eval_dataset = self.cfg['EvalDataset']
  166. eval_dataset.check_or_download_dataset()
  167. anno_file = eval_dataset.get_anno()
  168. self._metrics = [
  169. RBoxMetric(
  170. anno_file=anno_file,
  171. clsid2catid=clsid2catid,
  172. classwise=classwise,
  173. output_eval=output_eval,
  174. bias=bias,
  175. save_prediction_only=save_prediction_only)
  176. ]
  177. elif self.cfg.metric == 'VOC':
  178. self._metrics = [
  179. VOCMetric(
  180. label_list=self.dataset.get_label_list(),
  181. class_num=self.cfg.num_classes,
  182. map_type=self.cfg.map_type,
  183. classwise=classwise)
  184. ]
  185. elif self.cfg.metric == 'WiderFace':
  186. multi_scale = self.cfg.multi_scale_eval if 'multi_scale_eval' in self.cfg else True
  187. self._metrics = [
  188. WiderFaceMetric(
  189. image_dir=os.path.join(self.dataset.dataset_dir,
  190. self.dataset.image_dir),
  191. anno_file=self.dataset.get_anno(),
  192. multi_scale=multi_scale)
  193. ]
  194. elif self.cfg.metric == 'KeyPointTopDownCOCOEval':
  195. eval_dataset = self.cfg['EvalDataset']
  196. eval_dataset.check_or_download_dataset()
  197. anno_file = eval_dataset.get_anno()
  198. self._metrics = [
  199. KeyPointTopDownCOCOEval(anno_file,
  200. len(eval_dataset), self.cfg.num_joints,
  201. self.cfg.save_dir)
  202. ]
  203. elif self.cfg.metric == 'KeyPointTopDownMPIIEval':
  204. eval_dataset = self.cfg['EvalDataset']
  205. eval_dataset.check_or_download_dataset()
  206. anno_file = eval_dataset.get_anno()
  207. self._metrics = [
  208. KeyPointTopDownMPIIEval(anno_file,
  209. len(eval_dataset), self.cfg.num_joints,
  210. self.cfg.save_dir)
  211. ]
  212. elif self.cfg.metric == 'MOTDet':
  213. self._metrics = [JDEDetMetric(), ]
  214. else:
  215. logger.warning("Metric not support for metric type {}".format(
  216. self.cfg.metric))
  217. self._metrics = []
  218. def _reset_metrics(self):
  219. for metric in self._metrics:
  220. metric.reset()
  221. def register_callbacks(self, callbacks):
  222. callbacks = [c for c in list(callbacks) if c is not None]
  223. for c in callbacks:
  224. assert isinstance(c, Callback), \
  225. "metrics shoule be instances of subclass of Metric"
  226. self._callbacks.extend(callbacks)
  227. self._compose_callback = ComposeCallback(self._callbacks)
  228. def register_metrics(self, metrics):
  229. metrics = [m for m in list(metrics) if m is not None]
  230. for m in metrics:
  231. assert isinstance(m, Metric), \
  232. "metrics shoule be instances of subclass of Metric"
  233. self._metrics.extend(metrics)
  234. def load_weights(self, weights):
  235. if self.is_loaded_weights:
  236. return
  237. self.start_epoch = 0
  238. load_pretrain_weight(self.model, weights)
  239. logger.debug("Load weights {} to start training".format(weights))
  240. def load_weights_sde(self, det_weights, reid_weights):
  241. if self.model.detector:
  242. load_weight(self.model.detector, det_weights)
  243. load_weight(self.model.reid, reid_weights)
  244. else:
  245. load_weight(self.model.reid, reid_weights)
  246. def resume_weights(self, weights):
  247. # support Distill resume weights
  248. if hasattr(self.model, 'student_model'):
  249. self.start_epoch = load_weight(self.model.student_model, weights,
  250. self.optimizer)
  251. else:
  252. self.start_epoch = load_weight(self.model, weights, self.optimizer)
  253. logger.debug("Resume weights of epoch {}".format(self.start_epoch))
  254. def train(self, validate=False):
  255. assert self.mode == 'train', "Model not in 'train' mode"
  256. Init_mark = False
  257. # if validation in training is enabled, metrics should be re-init
  258. if validate:
  259. self._init_metrics(validate=validate)
  260. self._reset_metrics()
  261. model = self.model
  262. if self.cfg.get('fleet', False):
  263. model = fleet.distributed_model(model)
  264. self.optimizer = fleet.distributed_optimizer(self.optimizer)
  265. elif self._nranks > 1:
  266. find_unused_parameters = self.cfg[
  267. 'find_unused_parameters'] if 'find_unused_parameters' in self.cfg else False
  268. model = paddle.DataParallel(
  269. self.model, find_unused_parameters=find_unused_parameters)
  270. # initial fp16
  271. if self.cfg.get('fp16', False):
  272. scaler = amp.GradScaler(
  273. enable=self.cfg.use_gpu, init_loss_scaling=1024)
  274. self.status.update({
  275. 'epoch_id': self.start_epoch,
  276. 'step_id': 0,
  277. 'steps_per_epoch': len(self.loader)
  278. })
  279. self.status['batch_time'] = stats.SmoothedValue(
  280. self.cfg.log_iter, fmt='{avg:.4f}')
  281. self.status['data_time'] = stats.SmoothedValue(
  282. self.cfg.log_iter, fmt='{avg:.4f}')
  283. self.status['training_staus'] = stats.TrainingStats(self.cfg.log_iter)
  284. if self.cfg.get('print_flops', False):
  285. self._flops(self.loader)
  286. for epoch_id in range(self.start_epoch, self.cfg.epoch):
  287. self.status['mode'] = 'train'
  288. self.status['epoch_id'] = epoch_id
  289. self._compose_callback.on_epoch_begin(self.status)
  290. self.loader.dataset.set_epoch(epoch_id)
  291. model.train()
  292. iter_tic = time.time()
  293. for step_id, data in enumerate(self.loader):
  294. self.status['data_time'].update(time.time() - iter_tic)
  295. self.status['step_id'] = step_id
  296. self._compose_callback.on_step_begin(self.status)
  297. if self.cfg.get('fp16', False):
  298. with amp.auto_cast(enable=self.cfg.use_gpu):
  299. # model forward
  300. outputs = model(data)
  301. loss = outputs['loss']
  302. # model backward
  303. scaled_loss = scaler.scale(loss)
  304. scaled_loss.backward()
  305. # in dygraph mode, optimizer.minimize is equal to optimizer.step
  306. scaler.minimize(self.optimizer, scaled_loss)
  307. else:
  308. # model forward
  309. outputs = model(data)
  310. loss = outputs['loss']
  311. # model backward
  312. loss.backward()
  313. self.optimizer.step()
  314. curr_lr = self.optimizer.get_lr()
  315. self.lr.step()
  316. self.optimizer.clear_grad()
  317. self.status['learning_rate'] = curr_lr
  318. if self._nranks < 2 or self._local_rank == 0:
  319. self.status['training_staus'].update(outputs)
  320. self.status['batch_time'].update(time.time() - iter_tic)
  321. self._compose_callback.on_step_end(self.status)
  322. if self.use_ema:
  323. self.ema.update(self.model)
  324. iter_tic = time.time()
  325. # apply ema weight on model
  326. if self.use_ema:
  327. weight = copy.deepcopy(self.model.state_dict())
  328. self.model.set_dict(self.ema.apply())
  329. self._compose_callback.on_epoch_end(self.status)
  330. if validate and (self._nranks < 2 or self._local_rank == 0) \
  331. and ((epoch_id + 1) % self.cfg.snapshot_epoch == 0 \
  332. or epoch_id == self.end_epoch - 1):
  333. if not hasattr(self, '_eval_loader'):
  334. # build evaluation dataset and loader
  335. self._eval_dataset = self.cfg.EvalDataset
  336. self._eval_batch_sampler = \
  337. paddle.io.BatchSampler(
  338. self._eval_dataset,
  339. batch_size=self.cfg.EvalReader['batch_size'])
  340. self._eval_loader = create('EvalReader')(
  341. self._eval_dataset,
  342. self.cfg.worker_num,
  343. batch_sampler=self._eval_batch_sampler)
  344. # if validation in training is enabled, metrics should be re-init
  345. # Init_mark makes sure this code will only execute once
  346. if validate and Init_mark == False:
  347. Init_mark = True
  348. self._init_metrics(validate=validate)
  349. self._reset_metrics()
  350. with paddle.no_grad():
  351. self.status['save_best_model'] = True
  352. self._eval_with_loader(self._eval_loader)
  353. # restore origin weight on model
  354. if self.use_ema:
  355. self.model.set_dict(weight)
  356. def _eval_with_loader(self, loader):
  357. sample_num = 0
  358. tic = time.time()
  359. self._compose_callback.on_epoch_begin(self.status)
  360. self.status['mode'] = 'eval'
  361. self.model.eval()
  362. if self.cfg.get('print_flops', False):
  363. self._flops(loader)
  364. for step_id, data in enumerate(loader):
  365. self.status['step_id'] = step_id
  366. self._compose_callback.on_step_begin(self.status)
  367. # forward
  368. outs = self.model(data)
  369. # update metrics
  370. for metric in self._metrics:
  371. metric.update(data, outs)
  372. sample_num += data['im_id'].numpy().shape[0]
  373. self._compose_callback.on_step_end(self.status)
  374. self.status['sample_num'] = sample_num
  375. self.status['cost_time'] = time.time() - tic
  376. # accumulate metric to log out
  377. for metric in self._metrics:
  378. metric.accumulate()
  379. metric.log()
  380. self._compose_callback.on_epoch_end(self.status)
  381. # reset metric states for metric may performed multiple times
  382. self._reset_metrics()
  383. def evaluate(self):
  384. with paddle.no_grad():
  385. self._eval_with_loader(self.loader)
  386. def predict(self,
  387. images,
  388. draw_threshold=0.5,
  389. output_dir='output',
  390. save_txt=False):
  391. self.dataset.set_images(images)
  392. loader = create('TestReader')(self.dataset, 0)
  393. imid2path = self.dataset.get_imid2path()
  394. anno_file = self.dataset.get_anno()
  395. clsid2catid, catid2name = get_categories(
  396. self.cfg.metric, anno_file=anno_file)
  397. # Run Infer
  398. self.status['mode'] = 'test'
  399. self.model.eval()
  400. if self.cfg.get('print_flops', False):
  401. self._flops(loader)
  402. for step_id, data in enumerate(loader):
  403. self.status['step_id'] = step_id
  404. # forward
  405. outs = self.model(data)
  406. for key in ['im_shape', 'scale_factor', 'im_id']:
  407. outs[key] = data[key]
  408. for key, value in outs.items():
  409. if hasattr(value, 'numpy'):
  410. outs[key] = value.numpy()
  411. batch_res = get_infer_results(outs, clsid2catid)
  412. bbox_num = outs['bbox_num']
  413. start = 0
  414. for i, im_id in enumerate(outs['im_id']):
  415. image_path = imid2path[int(im_id)]
  416. image = Image.open(image_path).convert('RGB')
  417. self.status['original_image'] = np.array(image.copy())
  418. end = start + bbox_num[i]
  419. bbox_res = batch_res['bbox'][start:end] \
  420. if 'bbox' in batch_res else None
  421. mask_res = batch_res['mask'][start:end] \
  422. if 'mask' in batch_res else None
  423. segm_res = batch_res['segm'][start:end] \
  424. if 'segm' in batch_res else None
  425. keypoint_res = batch_res['keypoint'][start:end] \
  426. if 'keypoint' in batch_res else None
  427. image = visualize_results(
  428. image, bbox_res, mask_res, segm_res, keypoint_res,
  429. int(im_id), catid2name, draw_threshold)
  430. self.status['result_image'] = np.array(image.copy())
  431. if self._compose_callback:
  432. self._compose_callback.on_step_end(self.status)
  433. # save image with detection
  434. save_name = self._get_save_image_name(output_dir, image_path)
  435. logger.info("Detection bbox results save in {}".format(
  436. save_name))
  437. image.save(save_name, quality=95)
  438. if save_txt:
  439. save_path = os.path.splitext(save_name)[0] + '.txt'
  440. results = {}
  441. results["im_id"] = im_id
  442. if bbox_res:
  443. results["bbox_res"] = bbox_res
  444. if keypoint_res:
  445. results["keypoint_res"] = keypoint_res
  446. save_result(save_path, results, catid2name, draw_threshold)
  447. start = end
  448. def _get_save_image_name(self, output_dir, image_path):
  449. """
  450. Get save image name from source image path.
  451. """
  452. if not os.path.exists(output_dir):
  453. os.makedirs(output_dir)
  454. image_name = os.path.split(image_path)[-1]
  455. name, ext = os.path.splitext(image_name)
  456. return os.path.join(output_dir, "{}".format(name)) + ext
  457. def export(self, output_dir='output_inference'):
  458. self.model.eval()
  459. model_name = os.path.splitext(os.path.split(self.cfg.filename)[-1])[0]
  460. save_dir = os.path.join(output_dir, model_name)
  461. if not os.path.exists(save_dir):
  462. os.makedirs(save_dir)
  463. image_shape = None
  464. if self.cfg.architecture in MOT_ARCH:
  465. test_reader_name = 'TestMOTReader'
  466. else:
  467. test_reader_name = 'TestReader'
  468. if 'inputs_def' in self.cfg[test_reader_name]:
  469. inputs_def = self.cfg[test_reader_name]['inputs_def']
  470. image_shape = inputs_def.get('image_shape', None)
  471. # set image_shape=[3, -1, -1] as default
  472. if image_shape is None:
  473. image_shape = [3, -1, -1]
  474. self.model.eval()
  475. if hasattr(self.model, 'deploy'): self.model.deploy = True
  476. # Save infer cfg
  477. _dump_infer_config(self.cfg,
  478. os.path.join(save_dir, 'infer_cfg.yml'),
  479. image_shape, self.model)
  480. input_spec = [{
  481. "image": InputSpec(
  482. shape=[None] + image_shape, name='image'),
  483. "im_shape": InputSpec(
  484. shape=[None, 2], name='im_shape'),
  485. "scale_factor": InputSpec(
  486. shape=[None, 2], name='scale_factor')
  487. }]
  488. if self.cfg.architecture == 'DeepSORT':
  489. input_spec[0].update({
  490. "crops": InputSpec(
  491. shape=[None, 3, 192, 64], name='crops')
  492. })
  493. static_model = paddle.jit.to_static(self.model, input_spec=input_spec)
  494. # NOTE: dy2st do not pruned program, but jit.save will prune program
  495. # input spec, prune input spec here and save with pruned input spec
  496. pruned_input_spec = self._prune_input_spec(
  497. input_spec, static_model.forward.main_program,
  498. static_model.forward.outputs)
  499. # dy2st and save model
  500. if 'slim' not in self.cfg or self.cfg['slim_type'] != 'QAT':
  501. paddle.jit.save(
  502. static_model,
  503. os.path.join(save_dir, 'model'),
  504. input_spec=pruned_input_spec)
  505. else:
  506. self.cfg.slim.save_quantized_model(
  507. self.model,
  508. os.path.join(save_dir, 'model'),
  509. input_spec=pruned_input_spec)
  510. logger.info("Export model and saved in {}".format(save_dir))
  511. def _prune_input_spec(self, input_spec, program, targets):
  512. # try to prune static program to figure out pruned input spec
  513. # so we perform following operations in static mode
  514. paddle.enable_static()
  515. pruned_input_spec = [{}]
  516. program = program.clone()
  517. program = program._prune(targets=targets)
  518. global_block = program.global_block()
  519. for name, spec in input_spec[0].items():
  520. try:
  521. v = global_block.var(name)
  522. pruned_input_spec[0][name] = spec
  523. except Exception:
  524. pass
  525. paddle.disable_static()
  526. return pruned_input_spec
  527. def _flops(self, loader):
  528. self.model.eval()
  529. try:
  530. import paddleslim
  531. except Exception as e:
  532. logger.warning(
  533. 'Unable to calculate flops, please install paddleslim, for example: `pip install paddleslim`'
  534. )
  535. return
  536. from paddleslim.analysis import dygraph_flops as flops
  537. input_data = None
  538. for data in loader:
  539. input_data = data
  540. break
  541. input_spec = [{
  542. "image": input_data['image'][0].unsqueeze(0),
  543. "im_shape": input_data['im_shape'][0].unsqueeze(0),
  544. "scale_factor": input_data['scale_factor'][0].unsqueeze(0)
  545. }]
  546. flops = flops(self.model, input_spec) / (1000**3)
  547. logger.info(" Model FLOPs : {:.6f}G. (image shape is {})".format(
  548. flops, input_data['image'][0].unsqueeze(0).shape))