optimizer.py 42 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699700701702703704705706707708709710711712713714715716717718719720721722723724725726727728729730731732733734735736737738739740741742743744745746747748749750751752753754755756757758
  1. from __future__ import annotations
  2. import logging
  3. import os
  4. import time
  5. from functools import partial
  6. from typing import Callable, Optional, Sequence, Union
  7. import torch
  8. from hivemind.averaging.control import AveragingStage, StepControl
  9. from hivemind.compression import CompressionBase, NoCompression
  10. from hivemind.dht import DHT
  11. from hivemind.optim.experimental.grad_averager import GradientAverager
  12. from hivemind.optim.experimental.progress_tracker import ProgressTracker
  13. from hivemind.optim.experimental.state_averager import (
  14. LRSchedulerBase,
  15. OptimizerFactory,
  16. Parameters,
  17. ParamGroups,
  18. SchedulerFactory,
  19. TorchOptimizer,
  20. TrainingStateAverager,
  21. )
  22. from hivemind.optim.grad_scaler import GradScaler
  23. from hivemind.utils import PerformanceEMA, get_dht_time, get_logger
  24. logger = get_logger(__name__)
  25. class Optimizer(torch.optim.Optimizer):
  26. """
  27. hivemind.Optimizer wraps your regular PyTorch Optimizer for training collaboratively with peers.
  28. By default, Optimizer is configured to be exactly **equivalent to synchronous training** with target_batch_size.
  29. There are advanced options make training semi-asynchronous (delay_optimizer_step and delay_gradient_averaging)
  30. or even fully asynchronous (use_local_updates=True).
  31. :example: The Optimizer can be used as a drop-in replacement for a regular PyTorch Optimizer:
  32. >>> model = transformers.AutoModel("albert-xxlarge-v2")
  33. >>> dht = hivemind.DHT(initial_peers=INITIAL_PEERS, start=True)
  34. >>> opt = hivemind.Optimizer(dht=dht, run_id="run_42", batch_size_per_step=4, target_batch_size=4096,
  35. >>> params=model.parameters(), optimizer=lambda params: torch.optim.Adam(params))
  36. >>> while True:
  37. >>> loss = compute_loss_on_batch(model, batch_size=4)
  38. >>> opt.zero_grad()
  39. >>> loss.backward()
  40. >>> opt.step() # <-- train collaboratively with any peers that use the same prefix (run_42)
  41. By default, peers will perform the following steps:
  42. * accumulate a minibatch of gradients towards the (global) target batch size, without updating parameters yet;
  43. * after peers collectively accumulate target_batch_size, average gradients with peers and perform optimizer step;
  44. * if your peer lags behind the rest of the swarm, it will download parameters and optimizer state from others;
  45. Unlike regular training, your device may join midway through training, when other peers already made some progress.
  46. For this reason, any learning rate schedulers, curriculum and other **time-dependent features should be based on**
  47. ``optimizer.local_epoch`` (and not the number ot calls to opt.step). Otherwise, peers that joined training late
  48. may end up having different learning rates. To do so automatically, specify ``scheduler=...`` parameter below.
  49. :What is an epoch?: Optimizer uses the term ``epoch`` to describe intervals between synchronizations. One epoch
  50. coresponds to processing certain number of training samples (``target_batch_size``) in total across all peers.
  51. Like in PyTorch LR Scheduler, **epoch does not necessarily correspond to a full pass over the training data.**
  52. At the end of epoch, peers perform synchronous actions such as averaging gradients for a global optimizer update,
  53. updating the learning rate scheduler or simply averaging parameters (if using local updates).
  54. The purpose of this is to ensure that changing the number of peers does not require changing hyperparameters.
  55. For instance, if the number of peers doubles, they will run all-reduce more frequently to adjust for faster training.
  56. :Configuration guide: This guide will help you set up your first collaborative training run. It covers the most
  57. important basic options, but ignores features that require significant changes to the training code.
  58. >>> dht = hivemind.DHT(initial_peers=INITIAL_PEERS, client_mode=IF_BEHIND_FIREWALL_OR_VERY_UNRELIABLE, start=True)
  59. >>> opt = hivemind.Optimizer(
  60. >>> dht=dht, run_id="a_unique_name_that_every_participant_will_see_when_training",
  61. >>> batch_size_per_step=ACTUAL_BATCH_SIZE_OF_THIS_PEER, target_batch_size=LARGE_GLOBAL_BATCH,
  62. >>> # ^--- Each global optimzier step will use gradients from 1x-1.1x of target_batch_size (due to latency);
  63. >>> # It is recommended to train with very large batch sizes to reduce the % of time spent on communication.
  64. >>>
  65. >>> params=params, optimizer=lambda params: AnyPyTorchOptimizer(params, **hyperparams_for_target_batch_size),
  66. >>> # tune learning rate for your target_batch_size. Here's a good reference: https://arxiv.org/abs/1904.00962
  67. >>> scheduler=lambda opt: AnyPyTorchScheduler(opt, **hyperparams_for_target_batch_size),
  68. >>> # scheduler.step will be called automatically each time when peers collectively accumulate target_batch_size
  69. >>>
  70. >>> offload_optimizer=True, # saves GPU memory, but increases RAM usage; Generally a good practice to use this.
  71. >>> delay_grad_averaging=OPTIONAL, delay_optimizer_step=OPTIONAL, # train faster, but with 1 round of staleness;
  72. >>> # setting both to True is equivalent to Delayed Parameter Updates (see https://arxiv.org/abs/2101.06840)
  73. >>>
  74. >>> grad_compression=hivemind.Float16Compression(), state_averaging_compression=hivemind.Float16Compression(),
  75. >>> # ^-- it is usually fine to use pure 16-bit or even lower precision during communication with no precaution;
  76. >>> # See hivemind/examples/albert for an working example of mixed 8/16-bit compression.
  77. >>>
  78. >>> matchmaking_time=15.0, # 3-5s for small local runs, 10-15s for training over the internet or with many peers
  79. >>> averaging_timeout=60.0, # around of 2x the actual time it takes to run all-reduce
  80. >>> verbose=True # periodically report the training progress to the console (e.g. "Averaged with N peers")
  81. >>> ) # and you're done!
  82. :param dht: a running hivemind.DHT instance connected to other peers.
  83. :param run_id: a unique identifier of this training run, used as a common prefix for all DHT keys.
  84. **Note:** peers with the same run_id should *generally* train the same model and use compatible configurations.
  85. Some options can be safely changed by individual peers: ``batch_size_per_step``, ``client_mode``, ``auxiliary``,
  86. ``reuse_grad_buffers``, ``offload_optimizer``, and ``verbose``. In some cases, other options may also be tuned
  87. individually by each peer, but they should be changed with caution to avoid deadlocks or convergence issues.
  88. :param target_batch_size: global batch size that must be accumulated before the swarm transitions to the next epoch.
  89. The actual batch may be *slightly* larger due asynchrony (e.g. peers submit more gradients in the last second).
  90. :param batch_size_per_step: you should accumulate gradients over this many samples between calls to optimizer.step.
  91. :param params: parameters or param groups for the optimizer; required if optimizer is a callable(params).
  92. :param optimizer: a callable(parameters) -> pytorch.optim.Optimizer or a pre-initialized PyTorch optimizer.
  93. **Note:** some advanced options like offload_optimizer, delay_optimizer_step, or delay_grad_averaging require
  94. and require the callable and will not work if hivemind.optimizer is created with a pre-existing PyTorch Optimizer.
  95. :param scheduler: callable(optimizer) -> PyTorch LRScheduler or a pre-initialized PyTorch scheduler.
  96. The learning rate scheduler will adjust learning rate based on global epoch, not the number of
  97. local calls to optimizer.step; this is required to keep different peers synchronized.
  98. :param matchmaking_time: when looking for group, wait for peers to join for up to this many seconds.
  99. Increase if you see "averaged gradients with N peers" where N is below 0.9x the real siee on >=25% of epochs.
  100. When training with low-latency network, decreasing matchmaking_time allows training with smaller batch sizes.
  101. :param averaging_timeout: if an averaging step hangs for this long, it will be cancelled automatically.
  102. Increase averaging_timeout if you see "Proceeding with local gradients" at least 25% of the time.
  103. Do not set this timeout too high, as it may cause your optimizer to hang after some types of network errors.
  104. :param allreduce_timeout: timeout for a single attempt to run all-reduce, default: equal to averaging_timeout.
  105. :param load_state_timeout: wait for at most this many seconds before giving up on load_state_from_peers.
  106. :param reuse_grad_buffers: if True, use model's .grad buffers for gradient accumulation.
  107. This is more memory efficient, but it requires that the user does *NOT* call model/opt zero_grad at all
  108. :param offload_optimizer: offload the optimizer to host memory, saving GPU memory for parameters and gradients
  109. :param delay_optimizer_step: run optimizer in background, apply results in future .step; requires offload_optimizer
  110. :param delay_grad_averaging: average gradients in background; requires offload_optimizer and delay_optimizer_step
  111. :param delay_state_averaging: if enabled (default), average parameters and extra tensors in a background thread;
  112. if set to False, average parameters synchronously within the corresponding hivemind.Optimizer.step call.
  113. :param average_state_every: average state (parameters, chosen opt tensors) with peers every this many **epochs**.
  114. This reduces the communication overhead increasing, but can cause parameters to diverge if too large.
  115. The maximal average_state_every=num_epochs depends on how often peers diverge from each other. If peers
  116. hardly ever skip averaging rounds, they can average state less frequently. In turn, network failures, lossy
  117. gradient compression and local_updates cause parameters to diverge faster and requires more frequent averaging.
  118. :param use_local_updates: if enabled, peers will update parameters on each .step using local gradients;
  119. if not enabled (default), accumulate gradients to target_batch_size, and then call .step with averaged gradients.
  120. Even if use_local_updates=True, learning rate scheduler will still be called once per target_batch_size.
  121. :param client_mode: if True, this peer will not accept incoming connections (firewall-compatible mode)
  122. :param auxiliary: if True, optimizer.step will only assist other peers in averaging (for cpu-only workers)
  123. :param grad_compression: compression strategy used for averaging gradients, default = no compression
  124. :param state_averaging_compression: compression for averaging params and state tensors, default = no compression
  125. :param load_state_compression: compression strategy for loading state from peers, default = no compression
  126. :param average_opt_statistics: names of optimizer statistics from state dict that should be averaged with peers
  127. :param extra_tensors: if specified, these extra tensors will also be averaged and shared in load_state_from_peers.
  128. :param averager_opts: additional keyword arguments forwarded to both GradientAverager and TrainingStateAverager
  129. :param tracker_opts: additional keyword arguments forwarded to ProgressTracker
  130. :param performance_ema_alpha: moving average alpha in ProgressTracer, TrainingStateAverager and Optimizer
  131. :param verbose: if True, report internal events such as accumilating gradients and running background tasks
  132. :note: in a large-scale training, peers will inevitably fail and you will see error messages. hivemind.Optimizer
  133. is designed to recover from such failures, but will sometimes need a minute or two to re-adjust.
  134. """
  135. def __init__(
  136. self,
  137. *,
  138. dht: DHT,
  139. run_id: str,
  140. target_batch_size: int,
  141. batch_size_per_step: Optional[int] = None,
  142. optimizer: Union[TorchOptimizer, OptimizerFactory],
  143. params: Optional[Union[Parameters, ParamGroups]] = None,
  144. scheduler: Optional[Union[LRSchedulerBase, SchedulerFactory]] = None,
  145. matchmaking_time: Optional[float] = 15.0,
  146. averaging_timeout: Optional[float] = 60.0,
  147. allreduce_timeout: Optional[float] = None,
  148. next_chunk_timeout: Optional[float] = None,
  149. load_state_timeout: float = 600.0,
  150. reuse_grad_buffers: bool = False,
  151. offload_optimizer: Optional[bool] = None,
  152. delay_optimizer_step: Optional[bool] = None,
  153. delay_grad_averaging: bool = False,
  154. delay_state_averaging: bool = True,
  155. average_state_every: int = 1,
  156. use_local_updates: bool = False,
  157. client_mode: bool = None,
  158. auxiliary: bool = False,
  159. grad_compression: CompressionBase = NoCompression(),
  160. state_averaging_compression: CompressionBase = NoCompression(),
  161. load_state_compression: CompressionBase = NoCompression(),
  162. average_opt_statistics: Sequence[str] = (),
  163. extra_tensors: Sequence[torch.Tensor] = (),
  164. averager_opts: Optional[dict] = None,
  165. tracker_opts: Optional[dict] = None,
  166. performance_ema_alpha: float = 0.1,
  167. shutdown_timeout: float = 5,
  168. verbose: bool = False,
  169. ):
  170. client_mode = client_mode if client_mode is None else dht.client_mode
  171. delay_optimizer_step = delay_optimizer_step if delay_optimizer_step is not None else delay_grad_averaging
  172. offload_optimizer = offload_optimizer if offload_optimizer is not None else (params is not None)
  173. allreduce_timeout = allreduce_timeout if allreduce_timeout is not None else averaging_timeout
  174. next_chunk_timeout = next_chunk_timeout if next_chunk_timeout is not None else matchmaking_time
  175. assert not delay_grad_averaging or delay_optimizer_step, "delay_grad_averaging requires delay_optimizer_step"
  176. assert not (client_mode and auxiliary), "Client-mode peers cannot serve as auxiliaries"
  177. assert not auxiliary or batch_size_per_step is None, "Auxiliary peers should not accumulate batches"
  178. if callable(optimizer) and params is not None:
  179. if scheduler is not None and (not callable(scheduler) or isinstance(scheduler, LRSchedulerBase)):
  180. raise ValueError("For this mode, please provide scheduler factory: callable(optimizer) -> scheduler")
  181. elif all(hasattr(optimizer, attr) for attr in ("param_groups", "step", "zero_grad")):
  182. if offload_optimizer or delay_optimizer_step or delay_grad_averaging:
  183. raise ValueError(
  184. "To enable offload_optimizer or delayed updates, please initialize Optimizer as "
  185. "hivemind.Optimizer(..., params=params, optimizer=lambda params: create_opt(params)"
  186. )
  187. else:
  188. raise ValueError(
  189. "Please initialize the optimizer in one of the following two ways:\n"
  190. "(A) hivemind.Optimizer(..., params=params, optimizer=lambda params: create_opt(params)\n"
  191. "(B) hivemind.Optimizer(..., optimizer=pre_initialize_optimizer)"
  192. )
  193. if use_local_updates:
  194. assert not reuse_grad_buffers, "if local_updates is True, gradients will not be accumulated"
  195. assert not delay_grad_averaging, "if local_updates is True, gradients will not be averaged"
  196. self.dht, self.run_id, self.client_mode, self.auxiliary = dht, run_id, client_mode, auxiliary
  197. self.batch_size_per_step, self.target_batch_size = batch_size_per_step, target_batch_size
  198. self.delay_state_averaging, self.average_state_every = delay_state_averaging, average_state_every
  199. self.matchmaking_time, self.offload_optimizer = matchmaking_time, offload_optimizer
  200. self.delay_grad_averaging, self.delay_optimizer_step = delay_grad_averaging, delay_optimizer_step
  201. self.averaging_timeout, self.allreduce_timeout = averaging_timeout, allreduce_timeout
  202. self.load_state_timeout, self.shutdown_timeout = load_state_timeout, shutdown_timeout
  203. self.next_chunk_timeout = next_chunk_timeout
  204. self.status_loglevel = logging.INFO if verbose else logging.DEBUG
  205. self.scheduled_grads: Optional[StepControl] = None
  206. self.scheduled_state: Optional[StepControl] = None
  207. self.tracker = self._make_progress_tracker(
  208. target_batch_size, performance_ema_alpha=performance_ema_alpha, **tracker_opts or {}
  209. )
  210. self.state_averager = self._make_state_averager(
  211. optimizer=optimizer,
  212. params=params,
  213. scheduler=scheduler,
  214. delta_rule_averaging=use_local_updates and self.delay_state_averaging,
  215. compression=state_averaging_compression,
  216. state_compression=load_state_compression,
  217. average_opt_statistics=average_opt_statistics,
  218. performance_ema_alpha=performance_ema_alpha,
  219. extra_tensors=extra_tensors,
  220. **averager_opts or {},
  221. )
  222. if not use_local_updates:
  223. self.grad_averager = self._make_gradient_averager(
  224. reuse_grad_buffers=reuse_grad_buffers, compression=grad_compression, **averager_opts or {}
  225. )
  226. else:
  227. self.grad_averager = None
  228. self._should_check_synchronization_on_update = True # used in self.should_load_state_from_peers
  229. self._schema_hash = self._compute_schema_hash()
  230. self._parent_pid = os.getpid()
  231. self.delay_before_state_averaging = PerformanceEMA(alpha=performance_ema_alpha)
  232. # measures the average time from the beginning of self._update_global_epoch to the call to state_averager
  233. # used for pre-scheduling the averaging round in state_averager
  234. self._step_supports_amp_scaling = reuse_grad_buffers
  235. # note: the line above is used by pytorch AMP GradScaler to enable custom behavior needed when reusing gradient
  236. # buffers over multiple steps (to avoid repeated unscaling). Without reuse_grad_buffers, this is not needed.
  237. def _make_state_averager(self, **kwargs) -> TrainingStateAverager:
  238. return TrainingStateAverager(
  239. dht=self.dht,
  240. prefix=f"{self.run_id}_state_averager",
  241. min_matchmaking_time=self.matchmaking_time,
  242. allreduce_timeout=self.allreduce_timeout,
  243. shutdown_timeout=self.shutdown_timeout,
  244. offload_optimizer=self.offload_optimizer,
  245. custom_gradients=self.offload_optimizer,
  246. status_loglevel=self.status_loglevel,
  247. next_chunk_timeout=self.next_chunk_timeout,
  248. client_mode=self.client_mode,
  249. auxiliary=self.auxiliary,
  250. start=True,
  251. **kwargs,
  252. )
  253. def _make_gradient_averager(self, **kwargs) -> GradientAverager:
  254. assert hasattr(self, "state_averager"), "must initialize state averager first"
  255. grad_averager = GradientAverager(
  256. dht=self.dht,
  257. prefix=f"{self.run_id}_grad_averager",
  258. parameters=self.state_averager.main_parameters,
  259. min_matchmaking_time=self.matchmaking_time,
  260. allreduce_timeout=self.allreduce_timeout,
  261. shutdown_timeout=self.shutdown_timeout,
  262. next_chunk_timeout=self.next_chunk_timeout,
  263. client_mode=self.client_mode,
  264. auxiliary=self.auxiliary,
  265. start=True,
  266. **kwargs,
  267. )
  268. if self.offload_optimizer:
  269. optimized_param_groups = self.state_averager.optimizer.param_groups
  270. optimized_parameters = [param for group in optimized_param_groups for param in group["params"]]
  271. with grad_averager.get_tensors() as averaged_gradients:
  272. assert len(averaged_gradients) == len(optimized_parameters)
  273. for opt_param, averaged_grad in zip(optimized_parameters, averaged_gradients):
  274. opt_param.grad = averaged_grad
  275. return grad_averager
  276. def _make_progress_tracker(self, target_batch_size: int, **kwargs) -> ProgressTracker:
  277. return ProgressTracker(
  278. dht=self.dht,
  279. prefix=self.run_id,
  280. target_batch_size=target_batch_size,
  281. client_mode=self.client_mode,
  282. status_loglevel=self.status_loglevel,
  283. start=True,
  284. **kwargs,
  285. )
  286. def _compute_schema_hash(self) -> int:
  287. optimized_param_groups = self.state_averager.optimizer.param_groups
  288. optimized_parameters = [param for group in optimized_param_groups for param in group["params"]]
  289. param_shapes = tuple(tuple(param.shape) for param in optimized_parameters)
  290. # offloaded optimizer requires that gradient tensors are reused between iterations
  291. grad_ids = tuple(id(param.grad) for param in optimized_parameters) if self.offload_optimizer else None
  292. return hash((grad_ids, param_shapes))
  293. def is_alive(self) -> bool:
  294. return self.state_averager.is_alive()
  295. @property
  296. def local_epoch(self) -> int:
  297. """
  298. This worker's current epoch, kept synchronized with peers. If peer's local_epoch lags behind others, it will
  299. automatically re-synchronize by downloading state from another peer.
  300. An epoch corresponds to accumulating target_batch_size across all active devices.
  301. """
  302. return self.state_averager.local_epoch
  303. @property
  304. def use_local_updates(self) -> bool:
  305. return self.grad_averager is None
  306. @property
  307. def use_gradient_averaging(self) -> bool:
  308. return self.grad_averager is not None
  309. def step(
  310. self,
  311. closure: Optional[Callable[[], torch.Tensor]] = None,
  312. batch_size: Optional[int] = None,
  313. grad_scaler: Optional[GradScaler] = None,
  314. ):
  315. """
  316. Update training progress after accumulating another local batch size. Depending on the configuration, this will
  317. report progress to peers, run global or local optimizer step, average parameters or schedule background tasks.
  318. :param closure: A closure that reevaluates the model and returns the loss.
  319. :param batch_size: optional override for batch_size_per_step from init.
  320. :param grad_scaler: if amp is enabled, this **must** be a hivemind-aware gradient scaler.
  321. :note: this .step is different from normal pytorch optimizers in several key ways. See __init__ for details.
  322. """
  323. if grad_scaler is not None and not isinstance(grad_scaler, GradScaler):
  324. raise ValueError("hivemind.Optimizer requires a hivemind-aware gradient scaler (hivemind.GradScaler)")
  325. if self.batch_size_per_step is None and batch_size is None and not self.auxiliary:
  326. raise ValueError("Please either set batch_size_per_step parameter at init or when calling .step")
  327. if self.auxiliary and (closure is not None or batch_size is not None or grad_scaler is not None):
  328. raise ValueError("Auxiliary peers should not have batch size, run closures, or use grad_scaler")
  329. batch_size = batch_size if batch_size is not None else self.batch_size_per_step
  330. # if delayed updates finished before step, apply these updates; otherwise do nothing
  331. self.state_averager.step(apply_delayed_updates=True)
  332. loss = None
  333. if closure is not None:
  334. with torch.enable_grad():
  335. loss = closure()
  336. if not self.auxiliary and self.should_load_state_from_peers():
  337. logger.log(self.status_loglevel, "Peer is out of sync")
  338. self.load_state_from_peers()
  339. return loss # local gradients were computed with out-of-sync parameters, must start over
  340. if self.use_gradient_averaging:
  341. # accumulate gradients toward target batch size, then aggregate with peers and run optimizer
  342. if not self.auxiliary:
  343. grads_are_valid = self._check_and_accumulate_gradients(batch_size, grad_scaler)
  344. if not grads_are_valid:
  345. return loss # local gradients were reset due to overflow, must start over
  346. self._maybe_schedule_gradient_averaging()
  347. self._maybe_schedule_state_averaging()
  348. else:
  349. # use_local_updates=True: update parameters on every step independently of other peers
  350. if not self.auxiliary:
  351. if grad_scaler is not None:
  352. with grad_scaler.running_global_step():
  353. assert grad_scaler.unscale_(self)
  354. new_samples_accumulated = self.tracker.local_progress.samples_accumulated + batch_size
  355. self.tracker.report_local_progress(self.local_epoch, new_samples_accumulated)
  356. self._maybe_schedule_state_averaging()
  357. self.state_averager.step(
  358. increment_epoch=False,
  359. optimizer_step=True,
  360. delay_optimizer_step=self.delay_optimizer_step,
  361. grad_scaler=grad_scaler,
  362. )
  363. if self.tracker.ready_to_update_epoch:
  364. self._update_global_epoch(grad_scaler)
  365. return loss
  366. def _update_global_epoch(self, grad_scaler: Optional[GradScaler]) -> None:
  367. """Depending on the configuration: aggregate gradients and/or parameters, perform global optimizer step"""
  368. assert self._schema_hash == self._compute_schema_hash(), "parameters or gradients changed during iteration"
  369. _epoch_start_time = time.perf_counter()
  370. with self.tracker.pause_updates():
  371. wait_for_trigger = None
  372. if self.use_gradient_averaging:
  373. logger.log(self.status_loglevel, f"Beginning optimizer step #{self.local_epoch}")
  374. if self.delay_optimizer_step:
  375. self.state_averager.step(wait_for_delayed_updates=True)
  376. began_averaging_gradients = self._begin_averaging_gradients(grad_scaler)
  377. if not began_averaging_gradients:
  378. pass # failed to start gradient averaging due to an internal error
  379. elif self.delay_grad_averaging:
  380. # if using delayed grad averaing, send this to state_averager as a pre-condition for optimizer step
  381. wait_for_trigger = partial(self._average_gradients_and_load_into_optimizer, self.scheduled_grads)
  382. else:
  383. # delay_grad_averaging=False, average gradients immediately
  384. self._average_gradients_and_load_into_optimizer(self.scheduled_grads)
  385. next_epoch = max(self.local_epoch + 1, self.tracker.global_epoch)
  386. swarm_not_empty = self.tracker.global_progress.num_peers > 1
  387. should_perform_optimizer_step = not self.auxiliary and not self.use_local_updates
  388. should_average_state = (
  389. swarm_not_empty
  390. and next_epoch % self.average_state_every == 0
  391. and not self.state_averager.averaging_in_progress
  392. )
  393. if should_average_state and self.scheduled_state is not None:
  394. if self.scheduled_state.triggered or self.scheduled_state.done():
  395. logger.log(
  396. self.status_loglevel,
  397. f"Not using pre-scheduled group for state averaging because it"
  398. f"was already used elsewhere: {self.scheduled_state}",
  399. )
  400. self.scheduled_state = None
  401. self.delay_before_state_averaging.update(task_size=1, interval=time.perf_counter() - _epoch_start_time)
  402. self.state_averager.step(
  403. increment_epoch=True,
  404. wait_for_trigger=wait_for_trigger,
  405. optimizer_step=should_perform_optimizer_step,
  406. delay_optimizer_step=self.delay_optimizer_step and should_perform_optimizer_step,
  407. grad_scaler=grad_scaler,
  408. averaging_round=should_average_state,
  409. delay_averaging=self.delay_state_averaging and not self.auxiliary,
  410. averaging_control=self.scheduled_state if should_average_state else None,
  411. averaging_opts=dict(timeout=self.averaging_timeout) if should_average_state else None,
  412. )
  413. if not should_average_state and self.scheduled_state is not None and not self.scheduled_state.done():
  414. self.scheduled_state.cancel()
  415. self.scheduled_state = None
  416. self.tracker.update_epoch(new_epoch=self.state_averager.local_epoch)
  417. self._should_check_synchronization_on_update = True
  418. # the above line ensures that peers check for *strict* synchronization once per epoch
  419. if not self.client_mode:
  420. self.state_averager.state_sharing_priority = self.local_epoch
  421. if self.use_gradient_averaging and not self.auxiliary:
  422. self.grad_averager.reset_accumulated_grads_()
  423. if not self.client_mode:
  424. self.grad_averager.state_sharing_priority = self.local_epoch
  425. logger.log(self.status_loglevel, f"Transitioning to epoch {self.local_epoch}")
  426. def _begin_averaging_gradients(self, grad_scaler: Optional[GradScaler]) -> bool:
  427. """Begin an all-reduce round to average gradients; return True if succeeded, False if failed"""
  428. if grad_scaler is not None:
  429. with grad_scaler.running_global_step():
  430. assert grad_scaler.unscale_(self)
  431. began_averaging_gradients = False
  432. if self.scheduled_grads is not None and (self.scheduled_grads.triggered or self.scheduled_grads.done()):
  433. logger.log(
  434. self.status_loglevel,
  435. f"Not using pre-scheduled group for state averaging because it"
  436. f"was already used elsewhere: {self.scheduled_state}",
  437. )
  438. self.scheduled_grads = None
  439. elif self.tracker.global_progress.num_peers > 1:
  440. try:
  441. self.scheduled_grads = self.grad_averager.step(
  442. control=self.scheduled_grads, reset_accumulators=True, wait=False
  443. )
  444. began_averaging_gradients = True
  445. except BaseException as e:
  446. logger.exception(e)
  447. if not began_averaging_gradients and self.scheduled_grads is not None and not self.scheduled_grads.done():
  448. logger.log(self.status_loglevel, f"Tagging along for a pre-scheduled gradient averaging round")
  449. self._tag_along_with_zero_weight(self.scheduled_grads)
  450. self.scheduled_grads = None
  451. return began_averaging_gradients
  452. def _check_and_accumulate_gradients(self, batch_size: int, grad_scaler: Optional[GradScaler]) -> bool:
  453. """Check if gradients are valid, accumulate and return True; otherwise, reset and return False"""
  454. assert not self.use_local_updates and not self.auxiliary
  455. if grad_scaler is not None and not grad_scaler.are_grads_finite(self):
  456. logger.log(self.status_loglevel, "Encountered incorrect value in fp16 grads, resetting local gradients")
  457. self.tracker.report_local_progress(self.local_epoch, samples_accumulated=0)
  458. self.grad_averager.reset_accumulated_grads_()
  459. return False
  460. self.grad_averager.accumulate_grads_(batch_size)
  461. self.tracker.report_local_progress(self.local_epoch, self.grad_averager.local_samples_accumulated)
  462. return True
  463. def _maybe_schedule_gradient_averaging(self) -> None:
  464. """If next epoch is coming soon, schedule the next gradient averaging round at the estimated end of epoch"""
  465. assert self.use_gradient_averaging
  466. if self.tracker.estimated_next_update_time - get_dht_time() <= self.matchmaking_time:
  467. if self.scheduled_grads is None or self.scheduled_grads.triggered or self.scheduled_grads.done():
  468. eta_seconds = self.tracker.estimated_next_update_time - get_dht_time()
  469. eta_seconds = max(eta_seconds, self.grad_averager.matchmaking_kwargs["min_matchmaking_time"])
  470. logger.log(self.status_loglevel, f"Pre-scheduling gradient averaging round in {eta_seconds:.2f} sec")
  471. self.scheduled_grads = self.grad_averager.schedule_step(timeout=self.averaging_timeout)
  472. def _maybe_schedule_state_averaging(self) -> None:
  473. """If next epoch is coming soon, schedule the next state averaging at estimated parameter averaging start"""
  474. next_epoch = max(self.local_epoch + 1, self.tracker.global_epoch)
  475. if next_epoch % self.average_state_every != 0:
  476. return # averaging is not performed at this epoch
  477. if self.state_averager.averaging_in_progress:
  478. return # previous run is still in progress
  479. if self.delay_before_state_averaging.num_updates == 0:
  480. return # not enough data to accurately pre-schedule
  481. estimated_time = self.tracker.estimated_next_update_time
  482. estimated_time += self.delay_before_state_averaging.ema_seconds_per_sample
  483. estimated_time += self.state_averager.delay_before_averaging.ema_seconds_per_sample
  484. eta_seconds_to_averaging = estimated_time - get_dht_time()
  485. if eta_seconds_to_averaging <= self.matchmaking_time:
  486. if self.scheduled_state is None or self.scheduled_state.triggered or self.scheduled_state.done():
  487. min_matchmaking_time = self.state_averager.matchmaking_kwargs["min_matchmaking_time"]
  488. actual_seconds = max(eta_seconds_to_averaging, min_matchmaking_time)
  489. logger.log(self.status_loglevel, f"Pre-scheduling state averaging round in {actual_seconds:.2f} sec")
  490. self.scheduled_state = self.state_averager.schedule_step(
  491. gather=next_epoch, timeout=self.averaging_timeout
  492. )
  493. def _average_gradients_and_load_into_optimizer(self, maybe_step_control: Optional[StepControl]):
  494. """Run gradient averaging; on success, feed averaged gradients into optimizer; else, use local gradients"""
  495. assert self.use_gradient_averaging and maybe_step_control is None or maybe_step_control.triggered
  496. averaged_gradients = False
  497. try:
  498. if maybe_step_control is not None:
  499. group_info = maybe_step_control.result(self.averaging_timeout)
  500. logger.log(self.status_loglevel, f"Averaged gradients with {len(group_info)} peers")
  501. self._load_averaged_gradients_into_optimizer_()
  502. averaged_gradients = True
  503. else:
  504. logger.log(self.status_loglevel, f"Skipped averaging: there are no other peers")
  505. except BaseException as e:
  506. logger.log(self.status_loglevel, f"Averaging gradients failed with {repr(e)}")
  507. if not averaged_gradients:
  508. logger.log(self.status_loglevel, f"Proceeding with local gradients")
  509. self.grad_averager.load_accumulators_into_averager_()
  510. self._load_averaged_gradients_into_optimizer_()
  511. def _load_averaged_gradients_into_optimizer_(self):
  512. """If required, load averaged gradients into optimizer; otherwise simply notify grad averager"""
  513. assert self.use_gradient_averaging
  514. if self.offload_optimizer:
  515. pass # averaged gradients are already baked into optimizer, see _make_gradient_averager
  516. else:
  517. # copy averaged gradients into optimizer .grad buffers
  518. optimized_param_groups = self.state_averager.optimizer.param_groups
  519. optimized_parameters = [param for group in optimized_param_groups for param in group["params"]]
  520. with torch.no_grad(), self.grad_averager.get_tensors() as averaged_gradients:
  521. assert len(averaged_gradients) == len(optimized_parameters)
  522. for opt_param, averaged_grad in zip(optimized_parameters, averaged_gradients):
  523. opt_param.grad.copy_(averaged_grad, non_blocking=True)
  524. self.grad_averager.notify_used_averaged_gradients()
  525. def zero_grad(self, set_to_none: bool = False):
  526. """Reset gradients from model. If reuse_grad_buffers=True, this will raise an error."""
  527. if self.use_gradient_averaging and self.grad_averager.reuse_grad_buffers:
  528. raise ValueError(
  529. f"When running {self.__class__.__name__} with reuse_grad_buffers=True, user should never "
  530. f"call zero_grad manually. Gradients will be refreshed internally"
  531. )
  532. for param_group in self.param_groups:
  533. for param in param_group["params"]:
  534. if param.grad is None:
  535. pass
  536. elif set_to_none:
  537. param.grad = None
  538. else:
  539. param.grad.zero_()
  540. def should_load_state_from_peers(self) -> bool:
  541. """
  542. If true, peer will discard local progress and attempt to download state from peers.
  543. This method allows peer to continue training in two cases:
  544. - peer is on the same epoch as other collaborators - keep training normally
  545. - peer was on the same epoch and accumulated some grads, but some collaborators
  546. have just transitioned to the next epoch - this peer should also transition.
  547. :note: The latter case occurs due to the lack of network synchrony: the first peer that
  548. detects enough samples will transition to the next step and start counting samples anew.
  549. Some other peers may take time before they check with DHT and observe that
  550. - the global epoch is technically one epoch ahead of the current one and
  551. - the remaining (non-transitioned) peers no longer have target_batch_size between them
  552. If this is the case, peer should transition to the next epoch and does *not* need to re-load state.
  553. """
  554. if self._should_check_synchronization_on_update and self.tracker.fetched_global_progress_this_epoch.is_set():
  555. self._should_check_synchronization_on_update = False
  556. return self.local_epoch != self.tracker.global_epoch # require exact synchronization once per step
  557. return self.local_epoch < self.tracker.global_epoch - 1 # catch up if a peer just switched to next epoch
  558. def load_state_from_peers(self, **kwargs):
  559. """
  560. Attempt to load the newest collaboration state from other peers within the same run_id.
  561. If successful, this will update parameters, optimizer state, local epoch and learning rate schedule in-place.
  562. """
  563. # note: we tag along for the next all-reduce because the run may have already started and cancelling it
  564. # will cause peers to restart matchmaking and may stall the entire collaboration for a few seconds.
  565. if self.scheduled_grads is not None and not self.scheduled_grads.done():
  566. self._tag_along_with_zero_weight(self.scheduled_grads)
  567. self.scheduled_grads = None
  568. self.state_averager.step(wait_for_delayed_updates=True)
  569. with self.tracker.pause_updates():
  570. while True:
  571. try:
  572. self.state_averager.load_state_from_peers(timeout=self.load_state_timeout, **kwargs)
  573. break
  574. except KeyboardInterrupt:
  575. raise
  576. except BaseException as e:
  577. logger.exception(f"Failed to load state from peers: {e}, retrying ...")
  578. continue
  579. if self.tracker.global_epoch - 1 <= self.local_epoch < self.tracker.global_epoch:
  580. logger.log(self.status_loglevel, f"Catching up with collaboration step {self.tracker.global_epoch}")
  581. self.state_averager.local_epoch = self.tracker.global_epoch
  582. self.tracker.report_local_progress(local_epoch=self.local_epoch, samples_accumulated=0)
  583. if not self.client_mode:
  584. self.state_averager.state_sharing_priority = self.local_epoch
  585. if self.use_gradient_averaging:
  586. self.grad_averager.reset_accumulated_grads_()
  587. if not self.client_mode:
  588. self.grad_averager.state_sharing_priority = self.local_epoch
  589. def state_dict(self) -> dict:
  590. state_dict = self.state_averager.optimizer.state_dict()
  591. state_dict["state"]["local_epoch"] = self.local_epoch
  592. return state_dict
  593. def load_state_dict(self, state_dict: dict):
  594. if "local_epoch" in state_dict["state"]:
  595. self.state_averager.local_epoch = state_dict["state"].pop("local_epoch")
  596. return self.state_averager.optimizer.load_state_dict(state_dict)
  597. @property
  598. def state(self):
  599. return dict(self.state_averager.optimizer.state, local_epoch=self.local_epoch)
  600. @property
  601. def opt(self) -> TorchOptimizer:
  602. return self.state_averager.optimizer
  603. @property
  604. def param_groups(self) -> ParamGroups:
  605. next_index = 0
  606. param_groups = tuple(dict(param_group) for param_group in self.state_averager.optimizer.param_groups)
  607. for param_group in param_groups:
  608. num_params = len(param_group["params"])
  609. main_params_for_group = self.state_averager.main_parameters[next_index : next_index + num_params]
  610. param_group["params"] = main_params_for_group
  611. next_index += num_params
  612. assert next_index == len(self.state_averager.main_parameters)
  613. return param_groups
  614. def add_param_group(self, param_group: dict) -> None:
  615. raise ValueError(
  616. f"{self.__class__.__name__} does not support calling add_param_group after creation. "
  617. f"Please provide all parameter groups at init"
  618. )
  619. def __repr__(self):
  620. return f"{self.__class__.__name__}(prefix={self.run_id}, epoch={self.local_epoch})"
  621. def _tag_along_with_zero_weight(self, control: StepControl):
  622. """Wait for a running averaging round to finish with zero weight."""
  623. if not control.triggered:
  624. control.weight = 0
  625. control.allow_allreduce()
  626. if not control.done():
  627. try:
  628. control.result(self.averaging_timeout)
  629. except BaseException as e:
  630. logger.exception(e)
  631. if not control.done():
  632. control.cancel()
  633. def shutdown(self):
  634. logger.log(self.status_loglevel, "Sending goodbye to peers...")
  635. self.tracker.shutdown(self.shutdown_timeout)
  636. self.state_averager.step(wait_for_delayed_updates=True)
  637. for scheduled_round in self.scheduled_grads, self.scheduled_state:
  638. if scheduled_round is not None:
  639. if scheduled_round.stage == AveragingStage.LOOKING_FOR_GROUP:
  640. scheduled_round.cancel()
  641. else:
  642. self._tag_along_with_zero_weight(scheduled_round)
  643. logger.log(self.status_loglevel, "Shutting down averagers...")
  644. self.state_averager.shutdown()
  645. if self.use_gradient_averaging:
  646. self.grad_averager.shutdown()
  647. logger.log(self.status_loglevel, f"{self.__class__.__name__} is shut down")
  648. def __del__(self):
  649. if self._parent_pid == os.getpid() and self.is_alive():
  650. self.shutdown()