base.py 26 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633
  1. # Copyright (c) 2021 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. import os
  15. import os.path as osp
  16. from functools import partial
  17. import time
  18. import copy
  19. import math
  20. import yaml
  21. import json
  22. import numpy as np
  23. import paddle
  24. from paddle.io import DataLoader, DistributedBatchSampler
  25. from paddleslim import QAT
  26. from paddleslim.analysis import flops
  27. from paddleslim import L1NormFilterPruner, FPGMFilterPruner
  28. import paddlex
  29. from paddlex.cv.transforms import arrange_transforms
  30. from paddlex.utils import (seconds_to_hms, get_single_card_bs, dict2str,
  31. get_pretrain_weights, load_pretrain_weights,
  32. load_checkpoint, SmoothedValue, TrainingStats,
  33. _get_shared_memory_size_in_M, EarlyStop)
  34. import paddlex.utils.logging as logging
  35. from .slim.prune import _pruner_eval_fn, _pruner_template_input, sensitive_prune
  36. from .utils.infer_nets import InferNet
  37. class BaseModel:
  38. def __init__(self, model_type):
  39. self.model_type = model_type
  40. self.num_classes = None
  41. self.labels = None
  42. self.version = paddlex.__version__
  43. self.net = None
  44. self.optimizer = None
  45. self.test_inputs = None
  46. self.train_data_loader = None
  47. self.eval_data_loader = None
  48. self.eval_metrics = None
  49. # 是否使用多卡间同步BatchNorm均值和方差
  50. self.sync_bn = False
  51. self.status = 'Normal'
  52. # 已完成迭代轮数,为恢复训练时的起始轮数
  53. self.completed_epochs = 0
  54. self.pruner = None
  55. self.pruning_ratios = None
  56. self.quantizer = None
  57. self.quant_config = None
  58. self.fixed_input_shape = None
  59. def net_initialize(self,
  60. pretrain_weights=None,
  61. save_dir='.',
  62. resume_checkpoint=None,
  63. is_backbone_weights=False):
  64. if pretrain_weights is not None and \
  65. not osp.exists(pretrain_weights):
  66. if not osp.isdir(save_dir):
  67. if osp.exists(save_dir):
  68. os.remove(save_dir)
  69. os.makedirs(save_dir)
  70. if self.model_type == 'classifier':
  71. pretrain_weights = get_pretrain_weights(
  72. pretrain_weights, self.model_name, save_dir)
  73. else:
  74. backbone_name = getattr(self, 'backbone_name', None)
  75. pretrain_weights = get_pretrain_weights(
  76. pretrain_weights,
  77. self.__class__.__name__,
  78. save_dir,
  79. backbone_name=backbone_name)
  80. if pretrain_weights is not None:
  81. if is_backbone_weights:
  82. load_pretrain_weights(
  83. self.net.backbone,
  84. pretrain_weights,
  85. model_name='backbone of ' + self.model_name)
  86. else:
  87. load_pretrain_weights(
  88. self.net, pretrain_weights, model_name=self.model_name)
  89. if resume_checkpoint is not None:
  90. if not osp.exists(resume_checkpoint):
  91. logging.error(
  92. "The checkpoint path {} to resume training from does not exist."
  93. .format(resume_checkpoint),
  94. exit=True)
  95. if not osp.exists(osp.join(resume_checkpoint, 'model.pdparams')):
  96. logging.error(
  97. "Model parameter state dictionary file 'model.pdparams' "
  98. "not found under given checkpoint path {}".format(
  99. resume_checkpoint),
  100. exit=True)
  101. if not osp.exists(osp.join(resume_checkpoint, 'model.pdopt')):
  102. logging.error(
  103. "Optimizer state dictionary file 'model.pdparams' "
  104. "not found under given checkpoint path {}".format(
  105. resume_checkpoint),
  106. exit=True)
  107. if not osp.exists(osp.join(resume_checkpoint, 'model.yml')):
  108. logging.error(
  109. "'model.yml' not found under given checkpoint path {}".
  110. format(resume_checkpoint),
  111. exit=True)
  112. with open(osp.join(resume_checkpoint, "model.yml")) as f:
  113. info = yaml.load(f.read(), Loader=yaml.Loader)
  114. self.completed_epochs = info['completed_epochs']
  115. load_checkpoint(
  116. self.net,
  117. self.optimizer,
  118. model_name=self.model_name,
  119. checkpoint=resume_checkpoint)
  120. def get_model_info(self):
  121. info = dict()
  122. info['version'] = paddlex.__version__
  123. info['Model'] = self.__class__.__name__
  124. info['_Attributes'] = {'model_type': self.model_type}
  125. if 'self' in self.init_params:
  126. del self.init_params['self']
  127. if '__class__' in self.init_params:
  128. del self.init_params['__class__']
  129. if 'model_name' in self.init_params:
  130. del self.init_params['model_name']
  131. if 'params' in self.init_params:
  132. del self.init_params['params']
  133. info['_init_params'] = self.init_params
  134. info['_Attributes']['num_classes'] = self.num_classes
  135. info['_Attributes']['labels'] = self.labels
  136. info['_Attributes']['fixed_input_shape'] = self.fixed_input_shape
  137. try:
  138. primary_metric_key = list(self.eval_metrics.keys())[0]
  139. primary_metric_value = float(self.eval_metrics[primary_metric_key])
  140. info['_Attributes']['eval_metrics'] = {
  141. primary_metric_key: primary_metric_value
  142. }
  143. except:
  144. pass
  145. if hasattr(self, 'test_transforms'):
  146. if self.test_transforms is not None:
  147. info['Transforms'] = list()
  148. for op in self.test_transforms.transforms:
  149. name = op.__class__.__name__
  150. if name.startswith('Arrange'):
  151. continue
  152. attr = op.__dict__
  153. info['Transforms'].append({name: attr})
  154. info['completed_epochs'] = self.completed_epochs
  155. return info
  156. def get_pruning_info(self):
  157. info = dict()
  158. info['pruner'] = self.pruner.__class__.__name__
  159. info['pruning_ratios'] = self.pruning_ratios
  160. pruner_inputs = self.pruner.inputs
  161. if self.model_type == 'detector':
  162. pruner_inputs = {
  163. k: v.tolist()
  164. for k, v in pruner_inputs[0].items()
  165. }
  166. info['pruner_inputs'] = pruner_inputs
  167. return info
  168. def get_quant_info(self):
  169. info = dict()
  170. info['quant_config'] = self.quant_config
  171. return info
  172. def save_model(self, save_dir):
  173. if not osp.isdir(save_dir):
  174. if osp.exists(save_dir):
  175. os.remove(save_dir)
  176. os.makedirs(save_dir)
  177. model_info = self.get_model_info()
  178. model_info['status'] = self.status
  179. paddle.save(self.net.state_dict(),
  180. osp.join(save_dir, 'model.pdparams'))
  181. paddle.save(self.optimizer.state_dict(),
  182. osp.join(save_dir, 'model.pdopt'))
  183. with open(
  184. osp.join(save_dir, 'model.yml'), encoding='utf-8',
  185. mode='w') as f:
  186. yaml.dump(model_info, f)
  187. # 评估结果保存
  188. if hasattr(self, 'eval_details'):
  189. with open(osp.join(save_dir, 'eval_details.json'), 'w') as f:
  190. json.dump(self.eval_details, f)
  191. if self.status == 'Pruned' and self.pruner is not None:
  192. pruning_info = self.get_pruning_info()
  193. with open(
  194. osp.join(save_dir, 'prune.yml'), encoding='utf-8',
  195. mode='w') as f:
  196. yaml.dump(pruning_info, f)
  197. if self.status == 'Quantized' and self.quantizer is not None:
  198. quant_info = self.get_quant_info()
  199. with open(
  200. osp.join(save_dir, 'quant.yml'), encoding='utf-8',
  201. mode='w') as f:
  202. yaml.dump(quant_info, f)
  203. # 模型保存成功的标志
  204. open(osp.join(save_dir, '.success'), 'w').close()
  205. logging.info("Model saved in {}.".format(save_dir))
  206. def build_data_loader(self, dataset, batch_size, mode='train'):
  207. if dataset.num_samples < batch_size:
  208. raise Exception(
  209. 'The volume of dataset({}) must be larger than batch size({}).'
  210. .format(dataset.num_samples, batch_size))
  211. batch_size_each_card = get_single_card_bs(batch_size=batch_size)
  212. # TODO detection eval阶段需做判断
  213. batch_sampler = DistributedBatchSampler(
  214. dataset,
  215. batch_size=batch_size_each_card,
  216. shuffle=dataset.shuffle,
  217. drop_last=mode == 'train')
  218. if dataset.num_workers > 0:
  219. shm_size = _get_shared_memory_size_in_M()
  220. if shm_size is None or shm_size < 1024.:
  221. use_shared_memory = False
  222. else:
  223. use_shared_memory = True
  224. else:
  225. use_shared_memory = False
  226. loader = DataLoader(
  227. dataset,
  228. batch_sampler=batch_sampler,
  229. collate_fn=dataset.batch_transforms,
  230. num_workers=dataset.num_workers,
  231. return_list=True,
  232. use_shared_memory=use_shared_memory,
  233. worker_init_fn=lambda worker_id: np.random.seed(np.random.get_state()[1][0] + worker_id)
  234. )
  235. return loader
  236. def train_loop(self,
  237. num_epochs,
  238. train_dataset,
  239. train_batch_size,
  240. eval_dataset=None,
  241. save_interval_epochs=1,
  242. log_interval_steps=10,
  243. save_dir='output',
  244. ema=None,
  245. early_stop=False,
  246. early_stop_patience=5,
  247. use_vdl=True):
  248. arrange_transforms(
  249. model_type=self.model_type,
  250. transforms=train_dataset.transforms,
  251. mode='train')
  252. nranks = paddle.distributed.get_world_size()
  253. local_rank = paddle.distributed.get_rank()
  254. if nranks > 1:
  255. find_unused_parameters = getattr(self, 'find_unused_parameters',
  256. False)
  257. # Initialize parallel environment if not done.
  258. if not paddle.distributed.parallel.parallel_helper._is_parallel_ctx_initialized(
  259. ):
  260. paddle.distributed.init_parallel_env()
  261. ddp_net = paddle.DataParallel(
  262. self.net, find_unused_parameters=find_unused_parameters)
  263. else:
  264. ddp_net = paddle.DataParallel(
  265. self.net, find_unused_parameters=find_unused_parameters)
  266. if use_vdl:
  267. from visualdl import LogWriter
  268. vdl_logdir = osp.join(save_dir, 'vdl_log')
  269. log_writer = LogWriter(vdl_logdir)
  270. # task_id: 目前由PaddleX GUI赋值
  271. # 用于在VisualDL日志中注明所属任务id
  272. task_id = getattr(paddlex, "task_id", "")
  273. thresh = .0001
  274. if early_stop:
  275. earlystop = EarlyStop(early_stop_patience, thresh)
  276. self.train_data_loader = self.build_data_loader(
  277. train_dataset, batch_size=train_batch_size, mode='train')
  278. if eval_dataset is not None:
  279. self.test_transforms = copy.deepcopy(eval_dataset.transforms)
  280. start_epoch = self.completed_epochs
  281. train_step_time = SmoothedValue(log_interval_steps)
  282. train_step_each_epoch = math.floor(train_dataset.num_samples /
  283. train_batch_size)
  284. train_total_step = train_step_each_epoch * (num_epochs - start_epoch)
  285. if eval_dataset is not None:
  286. eval_batch_size = train_batch_size
  287. eval_epoch_time = 0
  288. best_accuracy_key = ""
  289. best_accuracy = -1.0
  290. best_model_epoch = -1
  291. current_step = 0
  292. for i in range(start_epoch, num_epochs):
  293. self.net.train()
  294. if callable(
  295. getattr(self.train_data_loader.dataset, 'set_epoch',
  296. None)):
  297. self.train_data_loader.dataset.set_epoch(i)
  298. train_avg_metrics = TrainingStats()
  299. step_time_tic = time.time()
  300. for step, data in enumerate(self.train_data_loader()):
  301. if nranks > 1:
  302. outputs = self.run(ddp_net, data, mode='train')
  303. else:
  304. outputs = self.run(self.net, data, mode='train')
  305. loss = outputs['loss']
  306. loss.backward()
  307. self.optimizer.step()
  308. self.optimizer.clear_grad()
  309. lr = self.optimizer.get_lr()
  310. if isinstance(self.optimizer._learning_rate,
  311. paddle.optimizer.lr.LRScheduler):
  312. self.optimizer._learning_rate.step()
  313. train_avg_metrics.update(outputs)
  314. outputs['lr'] = lr
  315. if ema is not None:
  316. ema.update(self.net)
  317. step_time_toc = time.time()
  318. train_step_time.update(step_time_toc - step_time_tic)
  319. step_time_tic = step_time_toc
  320. current_step += 1
  321. # 每间隔log_interval_steps,输出loss信息
  322. if current_step % log_interval_steps == 0 and local_rank == 0:
  323. if use_vdl:
  324. for k, v in outputs.items():
  325. log_writer.add_scalar(
  326. '{}-Metrics/Training(Step): {}'.format(
  327. task_id, k), v, current_step)
  328. # 估算剩余时间
  329. avg_step_time = train_step_time.avg()
  330. eta = avg_step_time * (train_total_step - current_step)
  331. if eval_dataset is not None:
  332. eval_num_epochs = math.ceil(
  333. (num_epochs - i - 1) / save_interval_epochs)
  334. if eval_epoch_time == 0:
  335. eta += avg_step_time * math.ceil(
  336. eval_dataset.num_samples / eval_batch_size)
  337. else:
  338. eta += eval_epoch_time * eval_num_epochs
  339. logging.info(
  340. "[TRAIN] Epoch={}/{}, Step={}/{}, {}, time_each_step={}s, eta={}"
  341. .format(i + 1, num_epochs, step + 1,
  342. train_step_each_epoch,
  343. dict2str(outputs),
  344. round(avg_step_time, 2), seconds_to_hms(eta)))
  345. logging.info('[TRAIN] Epoch {} finished, {} .'
  346. .format(i + 1, train_avg_metrics.log()))
  347. self.completed_epochs += 1
  348. # 每间隔save_interval_epochs, 在验证集上评估和对模型进行保存
  349. if ema is not None:
  350. weight = copy.deepcopy(self.net.state_dict())
  351. self.net.set_state_dict(ema.apply())
  352. eval_epoch_tic = time.time()
  353. if (i + 1) % save_interval_epochs == 0 or i == num_epochs - 1:
  354. if eval_dataset is not None and eval_dataset.num_samples > 0:
  355. eval_result = self.evaluate(
  356. eval_dataset,
  357. batch_size=eval_batch_size,
  358. return_details=True)
  359. # 保存最优模型
  360. if local_rank == 0:
  361. self.eval_metrics, self.eval_details = eval_result
  362. if use_vdl:
  363. for k, v in self.eval_metrics.items():
  364. try:
  365. log_writer.add_scalar(
  366. '{}-Metrics/Eval(Epoch): {}'.format(
  367. task_id, k), v, i + 1)
  368. except TypeError:
  369. pass
  370. logging.info('[EVAL] Finished, Epoch={}, {} .'.format(
  371. i + 1, dict2str(self.eval_metrics)))
  372. best_accuracy_key = list(self.eval_metrics.keys())[0]
  373. current_accuracy = self.eval_metrics[best_accuracy_key]
  374. if current_accuracy > best_accuracy:
  375. best_accuracy = current_accuracy
  376. best_model_epoch = i + 1
  377. best_model_dir = osp.join(save_dir, "best_model")
  378. self.save_model(save_dir=best_model_dir)
  379. if best_model_epoch > 0:
  380. logging.info(
  381. 'Current evaluated best model on eval_dataset is epoch_{}, {}={}'
  382. .format(best_model_epoch, best_accuracy_key,
  383. best_accuracy))
  384. eval_epoch_time = time.time() - eval_epoch_tic
  385. current_save_dir = osp.join(save_dir, "epoch_{}".format(i + 1))
  386. if local_rank == 0:
  387. self.save_model(save_dir=current_save_dir)
  388. if eval_dataset is not None and early_stop:
  389. if earlystop(current_accuracy):
  390. break
  391. if ema is not None:
  392. self.net.set_state_dict(weight)
  393. def analyze_sensitivity(self,
  394. dataset,
  395. batch_size=8,
  396. criterion='l1_norm',
  397. save_dir='output'):
  398. """
  399. Args:
  400. dataset(paddlex.dataset): Dataset used for evaluation during sensitivity analysis.
  401. batch_size(int, optional): Batch size used in evaluation. Defaults to 8.
  402. criterion({'l1_norm', 'fpgm'}, optional): Pruning criterion. Defaults to 'l1_norm'.
  403. save_dir(str, optional): The directory to save sensitivity file of the model. Defaults to 'output'.
  404. """
  405. if self.__class__.__name__ in ['FasterRCNN', 'MaskRCNN']:
  406. raise Exception("{} does not support pruning currently!".format(
  407. self.__class__.__name__))
  408. assert criterion in ['l1_norm', 'fpgm'], \
  409. "Pruning criterion {} is not supported. Please choose from ['l1_norm', 'fpgm']"
  410. arrange_transforms(
  411. model_type=self.model_type,
  412. transforms=dataset.transforms,
  413. mode='eval')
  414. if self.model_type == 'detector':
  415. self.net.eval()
  416. else:
  417. self.net.train()
  418. inputs = _pruner_template_input(
  419. sample=dataset[0], model_type=self.model_type)
  420. if criterion == 'l1_norm':
  421. self.pruner = L1NormFilterPruner(self.net, inputs=inputs)
  422. else:
  423. self.pruner = FPGMFilterPruner(self.net, inputs=inputs)
  424. if not osp.isdir(save_dir):
  425. os.makedirs(save_dir)
  426. sen_file = osp.join(save_dir, 'model.sensi.data')
  427. logging.info('Sensitivity analysis of model parameters starts...')
  428. self.pruner.sensitive(
  429. eval_func=partial(_pruner_eval_fn, self, dataset, batch_size),
  430. sen_file=sen_file)
  431. logging.info(
  432. 'Sensitivity analysis is complete. The result is saved at {}.'.
  433. format(sen_file))
  434. def prune(self, pruned_flops, save_dir=None):
  435. """
  436. Args:
  437. pruned_flops(float): Ratio of FLOPs to be pruned.
  438. save_dir(None or str, optional): If None, the pruned model will not be saved.
  439. Otherwise, the pruned model will be saved at save_dir. Defaults to None.
  440. """
  441. if self.status == "Pruned":
  442. raise Exception(
  443. "A pruned model cannot be done model pruning again!")
  444. pre_pruning_flops = flops(self.net, self.pruner.inputs)
  445. logging.info("Pre-pruning FLOPs: {}. Pruning starts...".format(
  446. pre_pruning_flops))
  447. _, self.pruning_ratios = sensitive_prune(self.pruner, pruned_flops)
  448. post_pruning_flops = flops(self.net, self.pruner.inputs)
  449. logging.info("Pruning is complete. Post-pruning FLOPs: {}".format(
  450. post_pruning_flops))
  451. logging.warning("Pruning the model may hurt its performance, "
  452. "retraining is highly recommended")
  453. self.status = 'Pruned'
  454. if save_dir is not None:
  455. self.save_model(save_dir)
  456. logging.info("Pruned model is saved at {}".format(save_dir))
  457. def _prepare_qat(self, quant_config):
  458. if self.status == 'Infer':
  459. logging.error(
  460. "Exported inference model does not support quantization aware training.",
  461. exit=True)
  462. if quant_config is None:
  463. # default quantization configuration
  464. quant_config = {
  465. # {None, 'PACT'}. Weight preprocess type. If None, no preprocessing is performed.
  466. 'weight_preprocess_type': None,
  467. # {None, 'PACT'}. Activation preprocess type. If None, no preprocessing is performed.
  468. 'activation_preprocess_type': None,
  469. # {'abs_max', 'channel_wise_abs_max', 'range_abs_max', 'moving_average_abs_max'}.
  470. # Weight quantization type.
  471. 'weight_quantize_type': 'channel_wise_abs_max',
  472. # {'abs_max', 'range_abs_max', 'moving_average_abs_max'}. Activation quantization type.
  473. 'activation_quantize_type': 'moving_average_abs_max',
  474. # The number of bits of weights after quantization.
  475. 'weight_bits': 8,
  476. # The number of bits of activation after quantization.
  477. 'activation_bits': 8,
  478. # Data type after quantization, such as 'uint8', 'int8', etc.
  479. 'dtype': 'int8',
  480. # Window size for 'range_abs_max' quantization.
  481. 'window_size': 10000,
  482. # Decay coefficient of moving average.
  483. 'moving_rate': .9,
  484. # Types of layers that will be quantized.
  485. 'quantizable_layer_type': ['Conv2D', 'Linear']
  486. }
  487. if self.status != 'Quantized':
  488. self.quant_config = quant_config
  489. self.quantizer = QAT(config=self.quant_config)
  490. logging.info(
  491. "Preparing the model for quantization-aware training...")
  492. self.quantizer.quantize(self.net)
  493. logging.info("Model is ready for quantization-aware training.")
  494. self.status = 'Quantized'
  495. elif quant_config != self.quant_config:
  496. logging.error(
  497. "The model has been quantized with the following quant_config: {}."
  498. "Doing quantization-aware training with a quantized model "
  499. "using a different configuration is not supported."
  500. .format(self.quant_config),
  501. exit=True)
  502. def _get_pipeline_info(self, save_dir):
  503. pipeline_info = {}
  504. pipeline_info["pipeline_name"] = self.model_type
  505. nodes = [{
  506. "src0": {
  507. "type": "Source",
  508. "next": "decode0"
  509. }
  510. }, {
  511. "decode0": {
  512. "type": "Decode",
  513. "next": "predict0"
  514. }
  515. }, {
  516. "predict0": {
  517. "type": "Predict",
  518. "init_params": {
  519. "use_gpu": False,
  520. "gpu_id": 0,
  521. "use_trt": False,
  522. "model_dir": save_dir,
  523. },
  524. "next": "sink0"
  525. }
  526. }, {
  527. "sink0": {
  528. "type": "Sink"
  529. }
  530. }]
  531. pipeline_info["pipeline_nodes"] = nodes
  532. pipeline_info["version"] = "1.0.0"
  533. return pipeline_info
  534. def _build_inference_net(self):
  535. infer_net = InferNet(self.net, self.model_type)
  536. infer_net.eval()
  537. return infer_net
  538. def _export_inference_model(self, save_dir, image_shape=None):
  539. save_dir = osp.join(save_dir, 'inference_model')
  540. self.test_inputs = self._get_test_inputs(image_shape)
  541. infer_net = self._build_inference_net()
  542. if self.status == 'Quantized':
  543. self.quantizer.save_quantized_model(infer_net,
  544. osp.join(save_dir, 'model'),
  545. self.test_inputs)
  546. quant_info = self.get_quant_info()
  547. with open(
  548. osp.join(save_dir, 'quant.yml'), encoding='utf-8',
  549. mode='w') as f:
  550. yaml.dump(quant_info, f)
  551. else:
  552. static_net = paddle.jit.to_static(
  553. infer_net, input_spec=self.test_inputs)
  554. paddle.jit.save(static_net, osp.join(save_dir, 'model'))
  555. if self.status == 'Pruned':
  556. pruning_info = self.get_pruning_info()
  557. with open(
  558. osp.join(save_dir, 'prune.yml'), encoding='utf-8',
  559. mode='w') as f:
  560. yaml.dump(pruning_info, f)
  561. model_info = self.get_model_info()
  562. model_info['status'] = 'Infer'
  563. with open(
  564. osp.join(save_dir, 'model.yml'), encoding='utf-8',
  565. mode='w') as f:
  566. yaml.dump(model_info, f)
  567. pipeline_info = self._get_pipeline_info(save_dir)
  568. with open(
  569. osp.join(save_dir, 'pipeline.yml'), encoding='utf-8',
  570. mode='w') as f:
  571. yaml.dump(pipeline_info, f)
  572. # 模型保存成功的标志
  573. open(osp.join(save_dir, '.success'), 'w').close()
  574. logging.info("The model for the inference deployment is saved in {}.".
  575. format(save_dir))