You can not select more than 25 topics Topics must start with a chinese character,a letter or number, can include dashes ('-') and can be up to 35 characters long.

_auto_parallel_context.py 34 kB

5 years ago
5 years ago
5 years ago
5 years ago
123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699700701702703704705706707708709710711712713714715716717718719720721722723724725726727728729730731732733734735736737738739740741742743744745746747748749750751752753754755756757758759760761762763764765766767768769770771772773774775776777778779780781782783784785786787788789790791792793794795796797798799
  1. # Copyright 2020 Huawei Technologies Co., Ltd
  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. # ============================================================================
  15. """Context of auto parallel"""
  16. import threading
  17. import mindspore.context as context
  18. import mindspore.log as logger
  19. from mindspore.parallel._dp_allreduce_fusion import _set_fusion_strategy_by_idx, _set_fusion_strategy_by_size
  20. from mindspore.parallel._ps_context import _is_role_pserver
  21. from mindspore._c_expression import AutoParallelContext
  22. from mindspore._checkparam import args_type_check, Validator
  23. _MAX_GROUP_NAME_LEN = 127
  24. _DEFAULT_HCCL_FUSION_GROUP_NAME = "hccl_world_groupsum1"
  25. _DEFAULT_NCCL_FUSION_GROUP_NAME = "nccl_world_groupsum1"
  26. class _AutoParallelContext:
  27. """
  28. _AutoParallelContext is the environment in which operations are executed
  29. Note:
  30. Create a context through instantiating Context object is not recommended.
  31. Should use auto_parallel_context() to get the context since Context is singleton.
  32. """
  33. _instance = None
  34. _instance_lock = threading.Lock()
  35. def __init__(self):
  36. self._context_handle = AutoParallelContext.get_instance()
  37. self._dataset_strategy_using_str = True
  38. def __new__(cls):
  39. if cls._instance is None:
  40. cls._instance_lock.acquire()
  41. cls._instance = object.__new__(cls)
  42. cls._instance_lock.release()
  43. return cls._instance
  44. def check_context_handle(self):
  45. """
  46. Check context handle.
  47. Raises:
  48. ValueError: If the context handle is none.
  49. """
  50. if self._context_handle is None:
  51. raise ValueError("Context handle is none in context!!!")
  52. def set_device_num(self, device_num):
  53. """
  54. Set device num for auto parallel.
  55. Args:
  56. device_num (int): The device number.
  57. Raises:
  58. ValueError: If the device num is not in [1, 4096].
  59. """
  60. self.check_context_handle()
  61. if device_num < 1 or device_num > 4096:
  62. raise ValueError("Device num must be in [1, 4096], but got {}".format(device_num))
  63. self._context_handle.set_device_num(device_num)
  64. def get_device_num(self):
  65. """Get device num."""
  66. self.check_context_handle()
  67. return self._context_handle.get_device_num()
  68. def set_global_rank(self, global_rank):
  69. """
  70. Set global rank for auto parallel.
  71. Args:
  72. global_rank (int): The rank id of current rank.
  73. Raises:
  74. ValueError: If the global rank is not in [1, 4096].
  75. """
  76. self.check_context_handle()
  77. if global_rank < 0 or global_rank > 4095:
  78. raise ValueError("Global rank must be in [0, 4095], but got {}".format(global_rank))
  79. self._context_handle.set_global_rank(global_rank)
  80. def get_global_rank(self):
  81. """Get current rank id."""
  82. self.check_context_handle()
  83. return self._context_handle.get_global_rank()
  84. def set_pipeline_stages(self, stages):
  85. """Set the stages of the pipeline"""
  86. if isinstance(stages, bool):
  87. raise TypeError("The type of pipeline_stage_num must be int, but got bool.")
  88. if not isinstance(stages, int):
  89. raise TypeError("The type of pipeline_stage_num must be int.")
  90. if stages < 1:
  91. raise ValueError("pipeline_stage_num can't be less than 1.")
  92. self.check_context_handle()
  93. self._context_handle.set_pipeline_stage_split_num(stages)
  94. def get_pipeline_stages(self):
  95. """Get the stages of the pipeline"""
  96. self.check_context_handle()
  97. return self._context_handle.get_pipeline_stage_split_num()
  98. def set_gradients_mean(self, gradients_mean):
  99. """
  100. Set gradients_mean flag.
  101. Note:
  102. If gradients_mean is true, it will insert a div operator after parameter gradients allreduce.
  103. Args:
  104. gradients_mean (bool): The gradients_mean flag.
  105. """
  106. self.check_context_handle()
  107. self._context_handle.set_gradients_mean(gradients_mean)
  108. def get_gradients_mean(self):
  109. """Get gradients_mean flag."""
  110. self.check_context_handle()
  111. return self._context_handle.get_gradients_mean()
  112. def set_gradient_fp32_sync(self, gradient_fp32_sync):
  113. """
  114. Set gradient_fp32_sync.
  115. Note:
  116. If gradient_fp32_sync is true,
  117. it will convert tensor type from fp16 to fp32 before parameter gradients allreduce.
  118. Args:
  119. gradient_fp32_sync (bool): The gradient_fp32_sync flag.
  120. """
  121. self.check_context_handle()
  122. self._context_handle.set_gradient_fp32_sync(gradient_fp32_sync)
  123. def get_gradient_fp32_sync(self):
  124. """Get gradient_fp32_sync flag."""
  125. self.check_context_handle()
  126. return self._context_handle.get_gradient_fp32_sync()
  127. def set_loss_repeated_mean(self, loss_repeated_mean):
  128. """
  129. Set loss_repeated_mean flag.
  130. Note:
  131. If loss_repeated_mean is true,
  132. Distributed automatic differentiation will perform a mean operator
  133. in backward in the case of repeated calculations.
  134. Args:
  135. loss_repeated_mean (bool): The loss_repeated_mean flag.
  136. """
  137. if not isinstance(loss_repeated_mean, bool):
  138. raise TypeError(f"The type of loss_repeated_mean must be bool, but got {type(loss_repeated_mean)}.")
  139. self.check_context_handle()
  140. self._context_handle.set_loss_repeated_mean(loss_repeated_mean)
  141. def get_loss_repeated_mean(self):
  142. """Get loss_repeated_mean flag."""
  143. self.check_context_handle()
  144. return self._context_handle.get_loss_repeated_mean()
  145. def set_parallel_mode(self, parallel_mode):
  146. """
  147. Set parallel mode for auto parallel.
  148. Args:
  149. parallel_mode (str): The parallel mode of auto parallel.
  150. Raises:
  151. ValueError: If parallel mode is not supported.
  152. """
  153. self.check_context_handle()
  154. ret = self._context_handle.set_parallel_mode(parallel_mode)
  155. if ret is False:
  156. raise ValueError("Parallel mode does not support {}".format(parallel_mode))
  157. def get_parallel_mode(self):
  158. """Get parallel mode."""
  159. self.check_context_handle()
  160. if _is_role_pserver():
  161. return context.ParallelMode.STAND_ALONE
  162. return self._context_handle.get_parallel_mode()
  163. def set_strategy_search_mode(self, auto_parallel_search_mode):
  164. """
  165. Set search mode of strategy.
  166. Args:
  167. auto_parallel_search_mode (str): The search mode of strategy.
  168. """
  169. self.check_context_handle()
  170. ret = self._context_handle.set_strategy_search_mode(auto_parallel_search_mode)
  171. if ret is False:
  172. raise ValueError("Strategy search mode does not support {}".format(auto_parallel_search_mode))
  173. def get_strategy_search_mode(self):
  174. """Get search mode of strategy."""
  175. self.check_context_handle()
  176. return self._context_handle.get_strategy_search_mode()
  177. def set_parameter_broadcast(self, parameter_broadcast):
  178. """
  179. Set parameter broadcast.
  180. Args:
  181. parameter_broadcast (bool): Parameter broadcast or not.
  182. """
  183. self.check_context_handle()
  184. self._context_handle.set_parameter_broadcast(parameter_broadcast)
  185. def get_parameter_broadcast(self):
  186. """Get parameter broadcast flag."""
  187. self.check_context_handle()
  188. return self._context_handle.get_parameter_broadcast()
  189. def set_strategy_ckpt_load_file(self, strategy_ckpt_load_file):
  190. """
  191. Set strategy checkpoint load path.
  192. Args:
  193. strategy_ckpt_load_file (str): Path to load parallel strategy checkpoint.
  194. """
  195. self.check_context_handle()
  196. self._context_handle.set_strategy_ckpt_load_file(strategy_ckpt_load_file)
  197. def get_strategy_ckpt_load_file(self):
  198. """Get strategy checkpoint load path."""
  199. self.check_context_handle()
  200. return self._context_handle.get_strategy_ckpt_load_file()
  201. def set_full_batch(self, full_batch):
  202. """
  203. Set whether load full batch on each device.
  204. Args:
  205. full_batch (bool): True if load full batch on each device.
  206. """
  207. self.check_context_handle()
  208. self._context_handle.set_full_batch(full_batch)
  209. def get_full_batch(self):
  210. """Get whether load full batch on each device."""
  211. self.check_context_handle()
  212. if _is_role_pserver():
  213. return False
  214. return self._context_handle.get_full_batch()
  215. def set_dataset_strategy(self, dataset_strategy):
  216. """
  217. Set dataset sharding strategy.
  218. Args:
  219. dataset_strategy (str or tuple(tuple)): The dataset sharding strategy.
  220. """
  221. self.check_context_handle()
  222. if isinstance(dataset_strategy, str):
  223. if dataset_strategy not in ("full_batch", "data_parallel"):
  224. raise ValueError("The dataset_strategy string should be 'full_batch' or 'data_parallel', "
  225. "otherwise, incoming tuple(tuple) type strategy")
  226. self._context_handle.set_full_batch(dataset_strategy == "full_batch")
  227. self._dataset_strategy_using_str = True
  228. return
  229. if not isinstance(dataset_strategy, tuple):
  230. raise TypeError(f'strategy must be str or tuple type, but got:{type(dataset_strategy)}')
  231. for ele in dataset_strategy:
  232. if not isinstance(ele, tuple):
  233. raise TypeError(f'The element of strategy must be tuple type, but got:{type(ele)}')
  234. for dim in ele:
  235. if not isinstance(dim, int):
  236. raise TypeError(f'The dim of each strategy value must be int type, but got:{type(dim)}')
  237. self._dataset_strategy_using_str = False
  238. self._context_handle.set_dataset_strategy(dataset_strategy)
  239. def get_dataset_strategy(self):
  240. """Get dataset sharding strategy."""
  241. self.check_context_handle()
  242. if self._dataset_strategy_using_str:
  243. if self._context_handle.get_full_batch():
  244. return "full_batch"
  245. return "data_parallel"
  246. return self._context_handle.get_dataset_strategy()
  247. def set_grad_accumulation_step(self, grad_accumulation_step):
  248. """
  249. Set grad accumulation step.
  250. Args:
  251. grad_accumulation_step (int): The grad accumulation step.
  252. """
  253. self.check_context_handle()
  254. Validator.check_positive_int(grad_accumulation_step)
  255. self._context_handle.set_grad_accumulation_step(grad_accumulation_step)
  256. def get_grad_accumulation_step(self):
  257. """Get grad accumulation step."""
  258. self.check_context_handle()
  259. return self._context_handle.get_grad_accumulation_step()
  260. def set_strategy_ckpt_save_file(self, strategy_ckpt_save_file):
  261. """
  262. Set strategy checkpoint save path.
  263. Args:
  264. strategy_ckpt_save_file (bool): Path to save parallel strategy checkpoint.
  265. """
  266. self.check_context_handle()
  267. import os
  268. dir_path = os.path.dirname(strategy_ckpt_save_file)
  269. if dir_path and not os.path.exists(dir_path):
  270. os.makedirs(dir_path)
  271. self._context_handle.set_strategy_ckpt_save_file(strategy_ckpt_save_file)
  272. def get_strategy_ckpt_save_file(self):
  273. """Get strategy checkpoint save path."""
  274. self.check_context_handle()
  275. return self._context_handle.get_strategy_ckpt_save_file()
  276. def set_group_ckpt_save_file(self, group_ckpt_save_file):
  277. """Set group checkpoint save path."""
  278. self.check_context_handle()
  279. import os
  280. dir_path = os.path.dirname(group_ckpt_save_file)
  281. if dir_path and not os.path.exists(dir_path):
  282. os.makedirs(dir_path)
  283. self._context_handle.set_group_ckpt_save_file(group_ckpt_save_file)
  284. def get_parameter_broadcast_is_set(self):
  285. """Get parameter broadcast is set or not."""
  286. self.check_context_handle()
  287. return self._context_handle.get_parameter_broadcast_is_set()
  288. def set_all_reduce_fusion_split_indices(self, indices, group=""):
  289. """
  290. Set allreduce fusion strategy by parameters indices.
  291. Args:
  292. indices (list): Indices list.
  293. group (str): The communication group of hccl/nccl.
  294. Raises:
  295. TypeError: If type of indices item is not int.
  296. TypeError: If group is not a python str.
  297. """
  298. self.check_context_handle()
  299. if not indices:
  300. raise ValueError('indices can not be empty')
  301. if isinstance(indices, (list)):
  302. for index in indices:
  303. if not isinstance(index, int) or isinstance(index, bool):
  304. raise TypeError(f"The type of index must be int, but got {type(index)}.")
  305. else:
  306. raise TypeError('indices must be a python list')
  307. if len(set(indices)) != len(indices):
  308. raise ValueError('indices has duplicate elements')
  309. if sorted(indices) != indices:
  310. raise ValueError('elements in indices must be sorted in ascending order')
  311. new_group = self._check_and_default_group(group)
  312. self._context_handle.set_all_reduce_fusion_split_indices(indices, new_group)
  313. if context.get_context("device_target") == "Ascend" and context.get_context("enable_ge"):
  314. _set_fusion_strategy_by_idx(indices)
  315. def get_all_reduce_fusion_split_indices(self, group=""):
  316. """
  317. Get allreduce fusion split indices.
  318. Args:
  319. group (str): The communication group of hccl/nccl.
  320. Returns:
  321. Return split sizes list according to the group.
  322. Raises:
  323. TypeError: If group is not a python str.
  324. """
  325. self.check_context_handle()
  326. new_group = self._check_and_default_group(group)
  327. return self._context_handle.get_all_reduce_fusion_split_indices(new_group)
  328. def set_all_reduce_fusion_split_sizes(self, sizes, group=""):
  329. """
  330. Set allreduce fusion strategy by parameters data sizes.
  331. Args:
  332. sizes (list): Sizes list.
  333. group (str): The communication group of hccl/nccl.
  334. Raises:
  335. TypeError: If type of sizes item is not int.
  336. TypeError: If group is not a python str.
  337. """
  338. self.check_context_handle()
  339. if isinstance(sizes, (list)):
  340. for size in sizes:
  341. if not isinstance(size, int) or isinstance(size, bool):
  342. raise TypeError(f"The type of size must be int, but got {type(size)}.")
  343. else:
  344. raise TypeError('sizes must be a python list')
  345. new_group = self._check_and_default_group(group)
  346. self._context_handle.set_all_reduce_fusion_split_sizes(sizes, new_group)
  347. if context.get_context("device_target") == "Ascend":
  348. _set_fusion_strategy_by_size(sizes)
  349. def get_all_reduce_fusion_split_sizes(self, group=""):
  350. """
  351. Get allreduce fusion split sizes.
  352. Args:
  353. group (str): The communication group of hccl/nccl.
  354. Returns:
  355. Return split sizes list according to the group.
  356. Raises:
  357. TypeError: If group is not a python str.
  358. """
  359. self.check_context_handle()
  360. new_group = self._check_and_default_group(group)
  361. return self._context_handle.get_all_reduce_fusion_split_sizes(new_group)
  362. def set_enable_all_reduce_fusion(self, enable_all_reduce_fusion):
  363. """
  364. Set enable/disable all reduce fusion.
  365. Args:
  366. enable_all_reduce_fusion (bool): Enable/disable all reduce fusion.
  367. """
  368. self.check_context_handle()
  369. if not isinstance(enable_all_reduce_fusion, bool):
  370. raise TypeError('enable_all_reduce_fusion is invalid type')
  371. self._context_handle.set_enable_all_reduce_fusion(enable_all_reduce_fusion)
  372. def get_enable_all_reduce_fusion(self):
  373. """Get all reduce fusion flag."""
  374. self.check_context_handle()
  375. return self._context_handle.get_enable_all_reduce_fusion()
  376. def get_device_num_is_set(self):
  377. """Get device number is set or not."""
  378. self.check_context_handle()
  379. return self._context_handle.get_device_num_is_set()
  380. def get_global_rank_is_set(self):
  381. """Get global rank is set or not."""
  382. self.check_context_handle()
  383. return self._context_handle.get_global_rank_is_set()
  384. def set_enable_parallel_optimizer(self, enable_parallel_optimizer):
  385. """
  386. Set enable/disable parallel optimizer.
  387. Args:
  388. set_enable_parallel_optimizer (bool): Enable/disable parallel optimizer.
  389. """
  390. self.check_context_handle()
  391. if not isinstance(enable_parallel_optimizer, bool):
  392. raise TypeError('enable_parallel_optimizer is invalid type')
  393. self._context_handle.set_enable_parallel_optimizer(enable_parallel_optimizer)
  394. def get_enable_parallel_optimizer(self):
  395. """Get parallel optimizer flag."""
  396. self.check_context_handle()
  397. return self._context_handle.get_enable_parallel_optimizer()
  398. def set_sharding_propagation(self, sharding_propagation):
  399. """
  400. Set the value of sharding strategy propagation in AUTO_PARALLEL mode. If True, the strategy-configured operators
  401. will propagate the strategies to other operators with minimum redistribution cost; otherwise, the algorithm
  402. will search the desired strategies.
  403. Default: False.
  404. Args:
  405. sharding_propagation (bool): Enable/disable strategy propagation.
  406. """
  407. self.check_context_handle()
  408. if not isinstance(sharding_propagation, bool):
  409. raise TypeError("'sharding_propagation' is an invalid type.")
  410. self._context_handle.set_sharding_propagation(sharding_propagation)
  411. def get_sharding_propagation(self):
  412. """Get the value of sharding strategy propagation."""
  413. self.check_context_handle()
  414. return self._context_handle.get_sharding_propagation()
  415. def set_enable_alltoall(self, enable_a2a):
  416. """
  417. Set the value of enabling AllToAll. If False, AllGather and Split are used to circumvent AllToAll.
  418. Default: False.
  419. Args:
  420. enable_a2a (bool): Enable/disable AllToAll.
  421. """
  422. self.check_context_handle()
  423. if not isinstance(enable_a2a, bool):
  424. raise TypeError("'enable_a2a' is an invalid type.")
  425. self._context_handle.set_enable_alltoall(enable_a2a)
  426. def get_enable_alltoall(self):
  427. """Get the value of enabling AllToAll."""
  428. self.check_context_handle()
  429. return self._context_handle.get_enable_alltoall()
  430. def set_communi_parallel_mode(self, communi_parallel_mode):
  431. """
  432. Set communication parallel mode.
  433. Args:
  434. communi_parallel_mode (str): The communication parallel mode.
  435. Raises:
  436. ValueError: If parallel mode is not supported.
  437. """
  438. if not isinstance(communi_parallel_mode, str):
  439. raise TypeError(f"The type of communi_parallel_mode must be str, \
  440. but got {type(communi_parallel_mode)}.")
  441. self.check_context_handle()
  442. ret = self._context_handle.set_communi_parallel_mode(communi_parallel_mode)
  443. if ret is False:
  444. raise ValueError("Communication parallel mode does not support {}".format(communi_parallel_mode))
  445. def get_communi_parallel_mode(self):
  446. """Get communication parallel mode."""
  447. self.check_context_handle()
  448. return self._context_handle.get_communi_parallel_mode()
  449. def set_optimizer_weight_shard_size(self, optimizer_weight_shard_size):
  450. """
  451. Set optimizer_weight_shard_size.
  452. Args:
  453. optimizer_weight_shard_size (int): Opt shard group size when not globally use parallel
  454. optimizer across devices.
  455. """
  456. self.check_context_handle()
  457. if not isinstance(optimizer_weight_shard_size, int) or isinstance(optimizer_weight_shard_size, bool):
  458. raise TypeError(f"The type of optimizer_weight_shard_size must be int, \
  459. but got {type(optimizer_weight_shard_size)}.")
  460. if optimizer_weight_shard_size <= 1:
  461. logger.warning("The setting 'optimizer_weight_shard_size' is invalid. "
  462. "Please use the integer larger than 1.")
  463. return
  464. self._context_handle.set_optimizer_weight_shard_size(optimizer_weight_shard_size)
  465. def get_optimizer_weight_shard_size(self):
  466. """Get optimizer_weight_shard_size."""
  467. self.check_context_handle()
  468. return self._context_handle.get_optimizer_weight_shard_size()
  469. def set_optimizer_weight_shard_aggregated_save(self, optimizer_weight_shard_aggregated_save):
  470. """
  471. Set optimizer_weight_shard_aggregated_save.
  472. Args:
  473. optimizer_weight_shard_aggregated_save (bool): Whether to integrated save weight shard when
  474. enable parallel optimizer.
  475. """
  476. self.check_context_handle()
  477. if not isinstance(optimizer_weight_shard_aggregated_save, bool):
  478. raise TypeError('optimizer_weight_shard_aggregated_save is invalid type')
  479. self._context_handle.set_optimizer_weight_shard_aggregated_save(optimizer_weight_shard_aggregated_save)
  480. def get_optimizer_weight_shard_aggregated_save(self):
  481. """Get optimizer_weight_shard_size."""
  482. self.check_context_handle()
  483. return self._context_handle.get_optimizer_weight_shard_aggregated_save()
  484. def reset(self):
  485. """Reset all settings."""
  486. self.check_context_handle()
  487. self._context_handle.reset()
  488. def _check_and_default_group(self, group):
  489. """Validate the given group, if group is empty, returns a default fusion group"""
  490. if isinstance(group, (str)):
  491. group_len = len(group)
  492. if group_len > _MAX_GROUP_NAME_LEN:
  493. raise ValueError('Group name len is out of range {_MAX_GROUP_NAME_LEN}')
  494. else:
  495. raise TypeError('Group must be a python str')
  496. if group == "":
  497. if context.get_context("device_target") == "Ascend":
  498. group = _DEFAULT_HCCL_FUSION_GROUP_NAME
  499. else:
  500. group = _DEFAULT_NCCL_FUSION_GROUP_NAME
  501. return group
  502. _auto_parallel_context = None
  503. def auto_parallel_context():
  504. """
  505. Get the global _auto_parallel_context, if it is not created, create a new one.
  506. Returns:
  507. _AutoParallelContext, the global auto parallel context.
  508. """
  509. global _auto_parallel_context
  510. if _auto_parallel_context is None:
  511. _auto_parallel_context = _AutoParallelContext()
  512. return _auto_parallel_context
  513. _set_auto_parallel_context_func_map = {
  514. "device_num": auto_parallel_context().set_device_num,
  515. "global_rank": auto_parallel_context().set_global_rank,
  516. "gradients_mean": auto_parallel_context().set_gradients_mean,
  517. "gradient_fp32_sync": auto_parallel_context().set_gradient_fp32_sync,
  518. "loss_repeated_mean": auto_parallel_context().set_loss_repeated_mean,
  519. "pipeline_stages": auto_parallel_context().set_pipeline_stages,
  520. "parallel_mode": auto_parallel_context().set_parallel_mode,
  521. "auto_parallel_search_mode": auto_parallel_context().set_strategy_search_mode,
  522. "parameter_broadcast": auto_parallel_context().set_parameter_broadcast,
  523. "strategy_ckpt_load_file": auto_parallel_context().set_strategy_ckpt_load_file,
  524. "strategy_ckpt_save_file": auto_parallel_context().set_strategy_ckpt_save_file,
  525. "group_ckpt_save_file": auto_parallel_context().set_group_ckpt_save_file,
  526. "full_batch": auto_parallel_context().set_full_batch,
  527. "dataset_strategy": auto_parallel_context().set_dataset_strategy,
  528. "enable_parallel_optimizer": auto_parallel_context().set_enable_parallel_optimizer,
  529. "grad_accumulation_step": auto_parallel_context().set_grad_accumulation_step,
  530. "all_reduce_fusion_config": auto_parallel_context().set_all_reduce_fusion_split_indices,
  531. "communi_parallel_mode": auto_parallel_context().set_communi_parallel_mode,
  532. "optimizer_weight_shard_size": auto_parallel_context().set_optimizer_weight_shard_size,
  533. "optimizer_weight_shard_aggregated_save": auto_parallel_context().set_optimizer_weight_shard_aggregated_save,
  534. "sharding_propagation": auto_parallel_context().set_sharding_propagation,
  535. "enable_alltoall": auto_parallel_context().set_enable_alltoall}
  536. _get_auto_parallel_context_func_map = {
  537. "device_num": auto_parallel_context().get_device_num,
  538. "global_rank": auto_parallel_context().get_global_rank,
  539. "gradients_mean": auto_parallel_context().get_gradients_mean,
  540. "gradient_fp32_sync": auto_parallel_context().get_gradient_fp32_sync,
  541. "loss_repeated_mean": auto_parallel_context().get_loss_repeated_mean,
  542. "pipeline_stages": auto_parallel_context().get_pipeline_stages,
  543. "parallel_mode": auto_parallel_context().get_parallel_mode,
  544. "auto_parallel_search_mode": auto_parallel_context().get_strategy_search_mode,
  545. "parameter_broadcast": auto_parallel_context().get_parameter_broadcast,
  546. "strategy_ckpt_load_file": auto_parallel_context().get_strategy_ckpt_load_file,
  547. "strategy_ckpt_save_file": auto_parallel_context().get_strategy_ckpt_save_file,
  548. "full_batch": auto_parallel_context().get_full_batch,
  549. "dataset_strategy": auto_parallel_context().get_dataset_strategy,
  550. "enable_parallel_optimizer": auto_parallel_context().get_enable_parallel_optimizer,
  551. "grad_accumulation_step": auto_parallel_context().get_grad_accumulation_step,
  552. "all_reduce_fusion_config": auto_parallel_context().get_all_reduce_fusion_split_indices,
  553. "communi_parallel_mode": auto_parallel_context().get_communi_parallel_mode,
  554. "optimizer_weight_shard_size": auto_parallel_context().get_optimizer_weight_shard_size,
  555. "optimizer_weight_shard_aggregated_save": auto_parallel_context().get_optimizer_weight_shard_aggregated_save,
  556. "sharding_propagation": auto_parallel_context().get_sharding_propagation,
  557. "enable_alltoall": auto_parallel_context().get_enable_alltoall}
  558. @args_type_check(device_num=int, global_rank=int, gradients_mean=bool, gradient_fp32_sync=bool,
  559. loss_repeated_mean=bool, parallel_mode=str, auto_parallel_search_mode=str,
  560. parameter_broadcast=bool, strategy_ckpt_load_file=str,
  561. strategy_ckpt_save_file=str, full_batch=bool, enable_parallel_optimizer=bool,
  562. grad_accumulation_step=int, all_reduce_fusion_config=list, group_ckpt_save_file=str,
  563. communi_parallel_mode=str, optimizer_weight_shard_size=int,
  564. optimizer_weight_shard_aggregated_save=bool,
  565. sharding_propagation=bool, enable_alltoall=bool)
  566. def _set_auto_parallel_context(**kwargs):
  567. """
  568. Set auto parallel context.
  569. Note:
  570. Attribute name is required for setting attributes.
  571. Args:
  572. device_num (int): Available device number, the value must be in [1, 4096]. Default: 1.
  573. global_rank (int): Global rank id, the value must be in [0, 4095]. Default: 0.
  574. gradients_mean (bool): Whether to perform mean operator after all-reduce of mirror. Default: False.
  575. loss_repeated_mean (bool): Whether to perform mean operator in backward in the case of repeated
  576. calculations. Default: True.
  577. gradient_fp32_sync (bool): Gradients allreduce by fp32 even though gradients is fp16 if this flag is True.
  578. Default: True.
  579. parallel_mode (str): There are five kinds of parallel modes, "stand_alone", "data_parallel",
  580. "hybrid_parallel", "semi_auto_parallel" and "auto_parallel". Default: "stand_alone".
  581. - stand_alone: Only one processor working.
  582. - data_parallel: Distributing the data across different processors.
  583. - hybrid_parallel: Achieving data parallelism and model parallelism manually.
  584. - semi_auto_parallel: Achieving data parallelism and model parallelism by
  585. setting parallel strategies.
  586. - auto_parallel: Achieving parallelism automatically.
  587. auto_parallel_search_mode (str): There are two kinds of search modes, "recursive_programming"
  588. and "dynamic_programming". Default: "dynamic_programming".
  589. - recursive_programming: Recursive programming search mode.
  590. - dynamic_programming: Dynamic programming search mode.
  591. parameter_broadcast (bool): Indicating whether to broadcast parameters before training.
  592. "stand_alone", "semi_auto_parallel" and "auto_parallel" do not support parameter
  593. broadcast. Default: False.
  594. strategy_ckpt_load_file (str): The path to load parallel strategy checkpoint. Default: ''
  595. strategy_ckpt_save_file (str): The path to save parallel strategy checkpoint. Default: ''
  596. group_ckpt_save_file (str): The path to save parallel group checkpoint. Default: ''
  597. full_batch (bool): Whether to load the whole batch on each device. Default: False.
  598. dataset_strategy Union[str, tuple]: Dataset sharding strategy. Default: "data_parallel".
  599. enable_parallel_optimizer (bool): Enable using optimizer segmentation or not. Default: False.
  600. all_reduce_fusion_config (list): Set allreduce fusion strategy by parameters indices.
  601. pipeline_stages (int): Set the stage information for pipeline parallel. This indicates how
  602. the devices are distributed alone the pipeline. The total devices will be divided into
  603. 'pipeline_stags' stages. This currently could only be used when
  604. parallel mode semi_auto_parallel is enabled. Default: 0
  605. communi_parallel_mode (str): There are tree kinds of communication parallel modes, "all_group_parallel",
  606. "same_server_group_parallel" and "no_group_parallel". Default: "all_group_parallel".
  607. - all_group_parallel: All communication groups are in parallel.
  608. - same_server_group_parallel: Only the communication groups within the same server are parallel.
  609. - no_group_parallel: All communication groups are not parallel.
  610. optimizer_weight_shard_size (int): Set optimizer shard group size when not fully use parallel optimizer.
  611. It should be larger than one and less than or equal with the data parallel size.
  612. Default: -1, which means fully use parallel optimizer in data parallel dimension.
  613. optimizer_weight_shard_aggregated_save (bool): Whether to integrated save weight shard when enable parallel
  614. optimizer. Default: False.
  615. sharding_propagation (bool): Set the value of sharding strategy propagation in AUTO_PARALLEL mode. If True,
  616. the strategy-configured operators will propagate the strategies to other
  617. operators with minimum redistribution cost; otherwise, the algorithm will
  618. search the desired strategies. Default: False.
  619. enable_alltoall (bool): Set the value of enabling AllToAll. If False, AllGather and Split are used to
  620. circumvent AllToAll. Default: False.
  621. Raises:
  622. ValueError: If input key is not attribute in auto parallel context.
  623. """
  624. for key, value in kwargs.items():
  625. if key not in _set_auto_parallel_context_func_map:
  626. raise ValueError("Set context keyword %s is not recognized!" % key)
  627. set_func = _set_auto_parallel_context_func_map[key]
  628. set_func(value)
  629. def _get_auto_parallel_context(attr_key):
  630. """
  631. Get auto parallel context attribute value according to the key.
  632. Args:
  633. attr_key (str): The key of the attribute.
  634. Returns:
  635. Return attribute value according to the key.
  636. Raises:
  637. ValueError: If input key is not attribute in auto parallel context.
  638. """
  639. if attr_key not in _get_auto_parallel_context_func_map:
  640. raise ValueError("Get context keyword %s is not recognized!" % attr_key)
  641. get_func = _get_auto_parallel_context_func_map[attr_key]
  642. return get_func()
  643. def _reset_auto_parallel_context():
  644. """
  645. Reset auto parallel context attributes to the default values:
  646. - device_num: 1.
  647. - global_rank: 0.
  648. - gradients_mean: False.
  649. - gradient_fp32_sync: True.
  650. - parallel_mode: "stand_alone".
  651. - parameter_broadcast: False.
  652. - strategy_ckpt_load_file: ""
  653. - strategy_ckpt_save_file: ""
  654. - enable_parallel_optimizer: False
  655. - auto_parallel_search_mode: dynamic_programming
  656. - pipeline_stages: 0
  657. """
  658. auto_parallel_context().reset()