averager.py 33 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699700701702703704705706707708709710711712713714715716717718719720721722723724725726727728
  1. """ A background process that averages your tensors with peers """
  2. from __future__ import annotations
  3. import asyncio
  4. import contextlib
  5. import ctypes
  6. import multiprocessing as mp
  7. import os
  8. import threading
  9. import weakref
  10. from dataclasses import asdict
  11. from typing import Any, AsyncIterator, Dict, Optional, Sequence, Tuple, Union
  12. import numpy as np
  13. import torch
  14. from hivemind.averaging.allreduce import AllreduceException, AllReduceRunner, AveragingMode, GroupID
  15. from hivemind.averaging.group_info import GroupInfo
  16. from hivemind.averaging.load_balancing import load_balance_peers
  17. from hivemind.averaging.matchmaking import Matchmaking, MatchmakingException
  18. from hivemind.averaging.partition import DEFAULT_PART_SIZE_BYTES
  19. from hivemind.compression import (
  20. CompressionBase,
  21. CompressionInfo,
  22. NoCompression,
  23. deserialize_torch_tensor,
  24. serialize_torch_tensor,
  25. )
  26. from hivemind.dht import DHT, DHTID
  27. from hivemind.p2p import P2PContext, P2PHandlerError, PeerID, ServicerBase
  28. from hivemind.proto import averaging_pb2
  29. from hivemind.utils import MPFuture, TensorDescriptor, get_logger
  30. from hivemind.utils.asyncio import achain, aiter_with_timeout, anext, as_aiter, azip, switch_to_uvloop
  31. from hivemind.utils.grpc import combine_from_streaming, split_for_streaming
  32. from hivemind.utils.serializer import MSGPackSerializer, SerializerBase
  33. from hivemind.utils.timed_storage import DHTExpiration, ValueWithExpiration, get_dht_time
  34. # flavour types
  35. GatheredData = Any
  36. logger = get_logger(__name__)
  37. class DecentralizedAverager(mp.Process, ServicerBase):
  38. """
  39. Parameter averaging service. A trainer can run this service in background to periodically average his parameters
  40. with other trainers. The averaging pattern is chosen so that (1) you only need to average with a small
  41. group of peers at a time, but (2) all trainers will converge to global average in a logarithmic number of steps.
  42. :param averaged_tensors: a sequence of pytorch tensors that will be averaged in each all-reduce
  43. :param dht: a DHT node that will be used to find groups
  44. :param start: if True, starts the background process immediately
  45. :param prefix: a shared prefix for all group keys
  46. :param target_group_size: attempts to form groups with up to this many peers (recommended: a power of 2, e.g. 16)
  47. :param initial_group_bits: a string of bits ('0' and '1') that define the initial group key (bucket index)
  48. :param averaging_expiration: attempt to find a group for this many seconds, otherwise try again
  49. note - this expiration time only applies to looking for group, passing tensors in allreduce may take more time
  50. :param compression: optionally compress tensors with this compression algorithm before running all-reduce
  51. :param state_compression: a separate compression strategy for load_state_from_peers (default = no compression)
  52. :param tensor_infos: CompressionInfo for each respective tensor; this determines how the tensor will be comressed
  53. :param allreduce_timeout: spend at most this many seconds for allreduce (after group is formed)
  54. :param averaging_alpha: optional "learning rate" for averaging. If specified, local parameters will be shifted
  55. towards the (estimated) average by this coefficient. By default, local parameters are set equal to average.
  56. :param request_timeout: when looking for group, wait for a response from leader for at most this many seconds.
  57. :note: request_timeout must be smaller than averaging_expiration to avoid potential deadlocks.
  58. :param part_size_bytes: tensors for AllReduce are processed in parts of up to this size (after compression)
  59. :param bandwidth: if specified, this value represents the network bandwidth available to averager.
  60. By default, the averager is assumed to have the average bandwidth of his group.
  61. If bandwidth == 0, averager will rely on its groupmates to do all the averaging.
  62. :param client_mode: if False, this averager will accept incoming requests from other peers.
  63. if True, the averager will only join existing groups where at least one peer has client_mode=False.
  64. By default, this flag is copied from DHTNode inside the ``dht`` instance.
  65. :param auxiliary: if this flag is specified, averager.step will only assist others without sending
  66. local tensors for averaging
  67. :param allow_state_sharing: if set to True, other peers can download this peer's state. Can be overwritten
  68. with averager.allow_state_sharing = True / False
  69. :param shutdown_timeout: when calling .shutdown, wait for up to this many seconds before terminating
  70. Example:
  71. >>> averager = DecentralizedAverager(...)
  72. >>> with averager.get_tensors() as tensors:
  73. >>> # run some code, modify tensors if necessary
  74. >>> tensors[0] += 1
  75. >>> # do not use tensors after the lock is released
  76. >>> metadata = averager.step(gather=dict(my_batch_size=32))
  77. >>> # run averaging once (in-place), gather metadata from groupmates
  78. >>> with averager.get_tensors() as tensors_after_averaging:
  79. >>> pass # use the averaged tensors
  80. """
  81. _matchmaking: Matchmaking
  82. _pending_group_assembled: asyncio.Event
  83. serializer = MSGPackSerializer
  84. def __init__(
  85. self,
  86. averaged_tensors: Sequence[torch.Tensor],
  87. dht: DHT,
  88. *,
  89. start: bool,
  90. prefix: str,
  91. target_group_size: int,
  92. min_group_size: int = 2,
  93. initial_group_bits: str = "",
  94. averaging_expiration: float = 15,
  95. request_timeout: float = 3,
  96. averaging_alpha: float = 1.0,
  97. part_size_bytes: int = DEFAULT_PART_SIZE_BYTES,
  98. allreduce_timeout: Optional[float] = None,
  99. compression: CompressionBase = NoCompression(),
  100. state_compression: CompressionBase = NoCompression(),
  101. tensor_infos: Optional[Sequence[CompressionInfo]] = None,
  102. bandwidth: Optional[float] = None,
  103. min_vector_size: int = 0,
  104. auxiliary: bool = False,
  105. allow_state_sharing: Optional[bool] = None,
  106. client_mode: Optional[bool] = None,
  107. daemon: bool = True,
  108. shutdown_timeout: float = 5,
  109. ):
  110. assert "." not in prefix, "group prefix must be a string without trailing '.'"
  111. assert bandwidth is None or (
  112. bandwidth >= 0 and np.isfinite(np.float32(bandwidth))
  113. ), "bandwidth must be a non-negative float32"
  114. if not is_power_of_two(target_group_size):
  115. logger.warning("It is recommended to set target_group_size to a power of 2.")
  116. assert all(bit in "01" for bit in initial_group_bits)
  117. assert not client_mode or not auxiliary, "auxiliary peers must accept incoming connections"
  118. super().__init__()
  119. self.dht = dht
  120. self.prefix = prefix
  121. if client_mode is None:
  122. client_mode = dht.client_mode
  123. self.client_mode = client_mode
  124. self._parent_pid = os.getpid()
  125. if self.client_mode:
  126. self.mode = AveragingMode.CLIENT
  127. elif auxiliary:
  128. self.mode = AveragingMode.AUX
  129. else:
  130. self.mode = AveragingMode.NODE
  131. self.daemon = daemon
  132. self._averaged_tensors = tuple(averaged_tensors)
  133. self.lock_averaged_tensors = mp.Lock()
  134. self.last_updated: DHTExpiration = -float("inf")
  135. for tensor in self._averaged_tensors:
  136. assert tensor.grad_fn is None, "averaged_tensors must be either parameters or leaf tensors"
  137. tensor.share_memory_()
  138. self.total_size = sum(map(torch.Tensor.numel, self._averaged_tensors))
  139. self.schema_hash = compute_schema_hash(self._averaged_tensors)
  140. self.shutdown_timeout = shutdown_timeout
  141. self.bandwidth = bandwidth
  142. self.matchmaking_kwargs = dict(
  143. servicer_type=type(self),
  144. prefix=prefix,
  145. initial_group_bits=initial_group_bits,
  146. target_group_size=target_group_size,
  147. min_group_size=min_group_size,
  148. averaging_expiration=averaging_expiration,
  149. request_timeout=request_timeout,
  150. )
  151. self.allreduce_kwargs = dict(
  152. compression=compression,
  153. part_size_bytes=part_size_bytes,
  154. min_vector_size=min_vector_size,
  155. )
  156. self._averaging_alpha, self._allreduce_timeout = averaging_alpha, allreduce_timeout
  157. self._running_groups: Dict[GroupID, AllReduceRunner] = {} # one or more assembled groups that run all-reduce
  158. self._inner_pipe, self._outer_pipe = mp.Pipe(duplex=True) # a control pipe used to communicate with daemon
  159. self._allow_state_sharing = mp.Value(ctypes.c_bool, 0)
  160. if allow_state_sharing is None:
  161. allow_state_sharing = not client_mode and not auxiliary
  162. self.allow_state_sharing = allow_state_sharing
  163. self.state_compression = state_compression
  164. self.tensor_infos = tensor_infos
  165. self._ready = MPFuture()
  166. # note: we create a background thread weakref and with daemon=True to ensure garbage collection
  167. background_fetcher = threading.Thread(
  168. daemon=True,
  169. target=_background_thread_fetch_current_state,
  170. args=[self.serializer, self._outer_pipe, weakref.WeakMethod(self.get_current_state)],
  171. )
  172. background_fetcher.start()
  173. if start:
  174. self.run_in_background(await_ready=True)
  175. @property
  176. def allow_state_sharing(self) -> bool:
  177. """if set to True, other peers can download this peer's state"""
  178. return bool(self._allow_state_sharing.value)
  179. @allow_state_sharing.setter
  180. def allow_state_sharing(self, value: bool):
  181. if value and self.client_mode:
  182. raise ValueError("Cannot allow state sharing: averager in client mode cannot share its state.")
  183. else:
  184. self._allow_state_sharing.value = value
  185. @property
  186. def peer_id(self) -> PeerID:
  187. return self.dht.peer_id
  188. @property
  189. def request_timeout(self):
  190. return self._matchmaking.request_timeout
  191. def run(self):
  192. """
  193. Run averager function in a background thread; this is needed to avoid a heisenbug with broken OMP on fork
  194. Turns out, using a non-main thread creates a separate OMP pool that works even if the original pool is corrupted
  195. Read more: https://github.com/pytorch/pytorch/issues/17199
  196. """
  197. thread = threading.Thread(target=self._run_internal, daemon=True)
  198. thread.start()
  199. thread.join()
  200. def _run_internal(self):
  201. """Serve DecentralizedAverager forever. This function will not return until the averager is shut down"""
  202. loop = switch_to_uvloop()
  203. # initialize asyncio synchronization primitives in this event loop
  204. pipe_semaphore = asyncio.Semaphore(value=0)
  205. loop.add_reader(self._inner_pipe.fileno(), pipe_semaphore.release)
  206. async def _run():
  207. try:
  208. self._p2p = await self.dht.replicate_p2p()
  209. if not self.client_mode:
  210. await self.add_p2p_handlers(self._p2p, namespace=self.prefix)
  211. else:
  212. logger.debug(f"The averager is running in client mode.")
  213. self._matchmaking = Matchmaking(
  214. self._p2p,
  215. self.schema_hash,
  216. self.dht,
  217. client_mode=self.client_mode,
  218. **self.matchmaking_kwargs,
  219. )
  220. if not self.client_mode:
  221. asyncio.create_task(self._declare_for_download_periodically())
  222. self._pending_group_assembled = asyncio.Event()
  223. self._pending_group_assembled.set()
  224. except Exception as e:
  225. # Loglevel is DEBUG since normally the exception is propagated to the caller
  226. logger.debug(e, exc_info=True)
  227. self._ready.set_exception(e)
  228. return
  229. self._ready.set_result(None)
  230. while True:
  231. try:
  232. await asyncio.wait_for(pipe_semaphore.acquire(), timeout=self.request_timeout)
  233. except asyncio.TimeoutError:
  234. pass
  235. if not self._inner_pipe.poll():
  236. continue
  237. try:
  238. method, args, kwargs = self._inner_pipe.recv()
  239. except (OSError, ConnectionError, RuntimeError) as e:
  240. logger.exception(e)
  241. await asyncio.sleep(self.request_timeout)
  242. continue
  243. task = asyncio.create_task(getattr(self, method)(*args, **kwargs))
  244. if method == "_shutdown":
  245. await task
  246. break
  247. loop.run_until_complete(_run())
  248. def run_in_background(self, await_ready: bool = True, timeout: Optional[float] = None) -> None:
  249. """
  250. Starts averager in a background process. if await_ready, this method will wait until background dht
  251. is ready to process incoming requests or for :timeout: seconds max.
  252. """
  253. self.start()
  254. if await_ready:
  255. self.wait_until_ready(timeout)
  256. def wait_until_ready(self, timeout: Optional[float] = None) -> None:
  257. self._ready.result(timeout=timeout)
  258. def shutdown(self) -> None:
  259. """Shut down the averager process"""
  260. if self.is_alive():
  261. self._outer_pipe.send(("_shutdown", [None], {})) # shut down the daemon process
  262. self._inner_pipe.send(("_SHUTDOWN", None)) # shut down background thread in master
  263. self.join(self.shutdown_timeout)
  264. if self.is_alive():
  265. logger.warning("Averager did not shut down within the grace period; terminating it the hard way.")
  266. self.terminate()
  267. else:
  268. logger.exception("Averager shutdown has no effect: the process is already not alive")
  269. async def _shutdown(self, timeout: Optional[float] = None) -> None:
  270. remaining_tasks = set()
  271. for group in self._running_groups.values():
  272. remaining_tasks.update(group.finalize(cancel=True))
  273. await asyncio.gather(*remaining_tasks)
  274. def __del__(self):
  275. if self._parent_pid == os.getpid() and self.is_alive():
  276. self.shutdown()
  277. def step(
  278. self,
  279. gather: Optional[GatheredData] = None,
  280. weight: Optional[float] = None,
  281. timeout: Optional[float] = None,
  282. allow_retries: bool = True,
  283. wait: bool = True,
  284. ) -> Union[Optional[Dict[PeerID, GatheredData]], MPFuture]:
  285. """
  286. Set up the averager to look for a group and run one round of averaging, return True on success, False on failure
  287. :param gather: optionally send this informaton to all peers in the next group and gather it from every groupmate
  288. (this operation is known as all-gather). The gathered data will be available as the output of this function.
  289. :param weight: averaging weight for this peer, int or float, must be strictly positive
  290. :param allow_retries: if averager fails to run one round of allreduce, this option will allow it to try again
  291. within the specified timeout
  292. :param timeout: if averager was unable to *find* a group in this many seconds, consider allreduce failedK
  293. :param wait: if True (default), return when finished. Otherwise return MPFuture and run in background.
  294. :returns: on success, update averaged_tensors and return group info; on failure, return None
  295. """
  296. if self.mode == AveragingMode.AUX and weight is not None:
  297. logger.warning("Averager is running in auxiliary mode, weight is unused.")
  298. if weight is None:
  299. weight = float(self.mode != AveragingMode.AUX)
  300. assert isinstance(weight, (int, float)) and weight >= 0, f"Expected a positive int/float, got {type(weight)}"
  301. future = MPFuture()
  302. gather_binary = self.serializer.dumps(
  303. gather
  304. ) # serialize here to avoid loading modules in the averager process
  305. self._outer_pipe.send(
  306. (
  307. "_step",
  308. [],
  309. dict(
  310. future=future,
  311. gather_binary=gather_binary,
  312. weight=weight,
  313. allow_retries=allow_retries,
  314. timeout=timeout,
  315. ),
  316. )
  317. )
  318. return future.result() if wait else future
  319. async def _step(
  320. self, *, future: MPFuture, gather_binary: bytes, weight: float, allow_retries: bool, timeout: Optional[float]
  321. ):
  322. start_time = get_dht_time()
  323. try:
  324. while not future.done():
  325. try:
  326. self._pending_group_assembled.clear()
  327. data_for_gather = self.serializer.dumps([self.bandwidth, self.mode.value, gather_binary])
  328. group_info = await self._matchmaking.look_for_group(
  329. timeout=timeout, data_for_gather=data_for_gather
  330. )
  331. if group_info is None:
  332. raise AllreduceException("Averaging step failed: could not find a group.")
  333. future.set_result(
  334. await asyncio.wait_for(
  335. self._run_allreduce(
  336. group_info, tensor_infos=self.tensor_infos, weight=weight, **self.allreduce_kwargs
  337. ),
  338. timeout=self._allreduce_timeout,
  339. )
  340. )
  341. # averaging is finished, loop will now exit
  342. except (
  343. AllreduceException,
  344. MatchmakingException,
  345. AssertionError,
  346. StopAsyncIteration,
  347. asyncio.CancelledError,
  348. asyncio.InvalidStateError,
  349. P2PHandlerError,
  350. ) as e:
  351. time_elapsed = get_dht_time() - start_time
  352. if not allow_retries or (timeout is not None and timeout < time_elapsed):
  353. logger.exception(f"Averager caught {repr(e)}")
  354. future.set_exception(e)
  355. else:
  356. logger.warning(f"Averager caught {repr(e)}, retrying")
  357. except BaseException as e:
  358. if not future.done():
  359. future.set_exception(e)
  360. raise
  361. finally:
  362. if not future.done():
  363. future.set_exception(
  364. RuntimeError(
  365. "Internal sanity check failed: averager.step left future pending."
  366. " Please report this to hivemind issues."
  367. )
  368. )
  369. async def _run_allreduce(self, group_info: GroupInfo, min_vector_size: int, **kwargs) -> GatheredData:
  370. """Run All-Reduce in a given group and update tensors in place, return gathered metadata"""
  371. try:
  372. bandwidths, mode_ids, user_gathered = zip(*map(self.serializer.loads, group_info.gathered))
  373. user_gathered = dict(zip(group_info.peer_ids, map(self.serializer.loads, user_gathered)))
  374. modes = tuple(map(AveragingMode, mode_ids))
  375. # compute optimal part sizes from peer bandwidths; TODO: replace with proper load balancing
  376. download_bandwidths = [
  377. thr if mode != AveragingMode.CLIENT else 0.0 for thr, mode in zip(bandwidths, modes)
  378. ]
  379. peer_fractions = await asyncio.get_event_loop().run_in_executor(
  380. None, load_balance_peers, self.total_size, download_bandwidths, min_vector_size
  381. )
  382. async with self.get_tensors_async() as local_tensors:
  383. allreduce = AllReduceRunner(
  384. p2p=self._p2p,
  385. servicer_type=type(self),
  386. prefix=self.prefix,
  387. group_id=group_info.group_id,
  388. tensors=local_tensors,
  389. ordered_peer_ids=group_info.peer_ids,
  390. peer_fractions=peer_fractions,
  391. gathered=user_gathered,
  392. modes=modes,
  393. **kwargs,
  394. )
  395. with self.register_allreduce_group(group_info.group_id, allreduce):
  396. if modes[group_info.peer_ids.index(self.peer_id)] != AveragingMode.AUX:
  397. async for tensor, update in azip(as_aiter(*local_tensors), allreduce):
  398. # all-reduce is performed asynchronously while iterating
  399. tensor.add_(update, alpha=self._averaging_alpha)
  400. self.last_updated = get_dht_time()
  401. else:
  402. async for _ in allreduce: # trigger all-reduce by iterating
  403. raise ValueError("aux peers should not receive averaged tensors")
  404. return allreduce.gathered
  405. except BaseException as e:
  406. logger.exception(e)
  407. raise MatchmakingException(f"Unable to run All-Reduce: {e}")
  408. @contextlib.contextmanager
  409. def register_allreduce_group(self, group_id: GroupID, allreduce: AllReduceRunner):
  410. """registers a given all-reduce runner to listen for incoming connections"""
  411. try:
  412. self._running_groups[group_id] = allreduce
  413. self._pending_group_assembled.set()
  414. yield
  415. finally:
  416. self._running_groups.pop(group_id, None)
  417. self._pending_group_assembled.set()
  418. @contextlib.contextmanager
  419. def get_tensors(self) -> Sequence[torch.Tensor]:
  420. """
  421. A contextmanager that gives user access to averaged tensors.
  422. It is guaranteed that the averager will not modify tensors while this context is active.
  423. Please do not modify the yielded tensors in-place after the context is released.
  424. """
  425. with self.lock_averaged_tensors:
  426. yield self._averaged_tensors
  427. self.last_updated = get_dht_time()
  428. @contextlib.asynccontextmanager
  429. async def get_tensors_async(self) -> Sequence[torch.Tensor]:
  430. """Like get_tensors, but uses an asynchronous contextmanager"""
  431. try:
  432. await asyncio.get_event_loop().run_in_executor(None, self.lock_averaged_tensors.acquire)
  433. yield self._averaged_tensors
  434. finally:
  435. self.lock_averaged_tensors.release()
  436. async def rpc_join_group(
  437. self, request: averaging_pb2.JoinRequest, context: P2PContext
  438. ) -> AsyncIterator[averaging_pb2.MessageFromLeader]:
  439. """accept or reject a join request from another averager; if accepted, run him through allreduce steps"""
  440. async for response in self._matchmaking.rpc_join_group(request, context):
  441. yield response
  442. async def rpc_aggregate_part(
  443. self, stream: AsyncIterator[averaging_pb2.AveragingData], context: P2PContext
  444. ) -> AsyncIterator[averaging_pb2.AveragingData]:
  445. """a groupmate sends us a part of his tensor; we should average it with other peers and return the result"""
  446. request = await anext(stream)
  447. if request.group_id not in self._running_groups:
  448. # this handles a special case when leader accepted us to group AND began allreduce right away,
  449. # but his response with group_id was delayed and other peers got to us first
  450. await self._pending_group_assembled.wait()
  451. group = self._running_groups.get(request.group_id)
  452. if group is None:
  453. yield averaging_pb2.AveragingData(code=averaging_pb2.BAD_GROUP_ID)
  454. return
  455. async for message in group.rpc_aggregate_part(achain(as_aiter(request), stream), context):
  456. yield message
  457. async def _declare_for_download_periodically(self):
  458. download_key = f"{self._matchmaking.group_key_manager.prefix}.all_averagers"
  459. while True:
  460. if self.allow_state_sharing:
  461. asyncio.create_task(
  462. asyncio.wait_for(
  463. self.dht.store(
  464. download_key,
  465. subkey=self.peer_id.to_bytes(),
  466. value=self.last_updated,
  467. expiration_time=get_dht_time() + self._matchmaking.averaging_expiration,
  468. return_future=True,
  469. ),
  470. timeout=self._matchmaking.averaging_expiration,
  471. )
  472. )
  473. await asyncio.sleep(self._matchmaking.averaging_expiration)
  474. async def rpc_download_state(
  475. self, _request: averaging_pb2.DownloadRequest, _context: P2PContext
  476. ) -> AsyncIterator[averaging_pb2.DownloadData]:
  477. """
  478. Get the up-to-date trainer state from a peer.
  479. The state consists of two parts: (serialized_metadata, tensors)
  480. - serialized_metadata is a small serialized bytestring meant to store scalars and hyperparameters
  481. - tensors is a sequence of pytorch tensors that represent model parameters or optimizer statistics
  482. """
  483. if not self.allow_state_sharing:
  484. return # deny request and direct peer to the next prospective averager
  485. metadata, tensors, infos = await self._get_current_state_from_host_process()
  486. if infos is None:
  487. infos = [CompressionInfo.from_tensor(tensor, key=i) for i, tensor in enumerate(tensors)]
  488. assert len(tensors) == len(infos)
  489. for tensor, info in zip(tensors, infos):
  490. for part in split_for_streaming(self.state_compression.compress(tensor, info, allow_inplace=False)):
  491. if metadata is not None:
  492. yield averaging_pb2.DownloadData(tensor_part=part, metadata=metadata)
  493. metadata = None
  494. else:
  495. yield averaging_pb2.DownloadData(tensor_part=part)
  496. def get_current_state(self) -> Tuple[Any, Sequence[torch.Tensor], Sequence[CompressionInfo]]:
  497. """
  498. Get current state and send it to a peer. executed in the host process. Meant to be overriden.
  499. :returns: a tuple of (small metadata, sequence of torch tensors)
  500. :note: metadata must be seriablizable with self.serializer (default = MSGPackSerializer)
  501. """
  502. with self.get_tensors() as tensors:
  503. return dict(group_key=self.get_group_bits()), tensors, self.tensor_infos
  504. async def _get_current_state_from_host_process(self):
  505. """Executed in the averager process inside rpc_download_state"""
  506. future = MPFuture()
  507. self._inner_pipe.send(("_TRIGGER_GET_CURRENT_STATE", future))
  508. return await future
  509. def load_state_from_peers(
  510. self, wait: bool = True, timeout: Optional[float] = None
  511. ) -> Optional[Tuple[Any, Sequence[torch.Tensor]]]:
  512. """
  513. Try to download the latest optimizer state one of the existing peer.
  514. :returns: on success, return a 2-tuple with (metadata, tensors), where
  515. - metadata is a small object containing metadata (e.g. hyperparameters, scalars, etc)
  516. - tensors is a sequence of pytorch tensors meant to contain peer's model weights and optimizer statistics
  517. The exact contents of both metadata and tensors are determined by get_current_state method
  518. """
  519. future = MPFuture()
  520. self._outer_pipe.send(("_load_state_from_peers", [], dict(future=future)))
  521. return future.result(timeout=timeout) if wait else future
  522. async def _load_state_from_peers(self, future: MPFuture):
  523. try:
  524. key_manager = self._matchmaking.group_key_manager
  525. peer_priority, _ = self.dht.get(f"{key_manager.prefix}.all_averagers", latest=True) or ({}, None)
  526. peer_priority = {
  527. PeerID(peer_id): float(info.value)
  528. for peer_id, info in peer_priority.items()
  529. if isinstance(info, ValueWithExpiration) and isinstance(info.value, (float, int))
  530. }
  531. if not isinstance(peer_priority, dict) or len(peer_priority) == 0:
  532. logger.info(f"Averager could not load state from peers: peer dict empty or corrupted {peer_priority}.")
  533. future.set_result(None)
  534. return
  535. metadata = None
  536. for peer in sorted(peer_priority.keys(), key=peer_priority.get, reverse=True):
  537. if peer != self.peer_id:
  538. logger.info(f"Downloading parameters from peer {peer}")
  539. try:
  540. stub = self.get_stub(self._p2p, peer, namespace=self.prefix)
  541. stream = await stub.rpc_download_state(averaging_pb2.DownloadRequest())
  542. current_tensor_parts, tensors = [], []
  543. async for message in aiter_with_timeout(stream, timeout=self.request_timeout):
  544. if message.metadata:
  545. metadata = self.serializer.loads(message.metadata)
  546. if message.tensor_part.dtype and current_tensor_parts:
  547. # tensor_part.dtype indicates the start of the new tensor, so we should wrap up this one
  548. tensors.append(deserialize_torch_tensor(combine_from_streaming(current_tensor_parts)))
  549. current_tensor_parts = []
  550. current_tensor_parts.append(message.tensor_part)
  551. if current_tensor_parts:
  552. tensors.append(deserialize_torch_tensor(combine_from_streaming(current_tensor_parts)))
  553. if not metadata:
  554. logger.debug(f"Peer {peer} did not send its state.")
  555. continue
  556. logger.info(f"Finished downloading state from {peer}")
  557. future.set_result((metadata, tensors))
  558. self.last_updated = get_dht_time()
  559. return
  560. except Exception as e:
  561. logger.exception(f"Failed to download state from {peer} - {repr(e)}")
  562. finally:
  563. if not future.done():
  564. future.set_result(None)
  565. def get_group_bits(self, wait: bool = True):
  566. """
  567. :param wait: if True, return bits immediately. Otherwise return awaitable MPFuture
  568. :returns: averager's current group key bits (without prefix)
  569. """
  570. future = MPFuture()
  571. self._outer_pipe.send(("_get_group_bits", [], dict(future=future)))
  572. return future.result() if wait else future
  573. async def _get_group_bits(self, future: MPFuture):
  574. future.set_result(self._matchmaking.group_key_manager.group_bits)
  575. def set_group_bits(self, group_bits: str, wait: bool = True):
  576. """
  577. :param group_bits: group bits (string of '0' or '1') to be used in averager's group key
  578. :param wait: if True, wait until the update is confirmed by the averager. Otherwise return immediately
  579. """
  580. future = MPFuture()
  581. assert all(bit in "01" for bit in group_bits)
  582. self._outer_pipe.send(("_set_group_bits", [], dict(group_bits=group_bits, future=future)))
  583. return future.result() if wait else future
  584. async def _set_group_bits(self, group_bits: str, future: MPFuture):
  585. try:
  586. self._matchmaking.group_key_manager.group_bits = group_bits
  587. return future.set_result(None)
  588. except Exception as e:
  589. if not future.done():
  590. future.set_exception(e)
  591. def is_power_of_two(n):
  592. """Check whether n is a power of 2"""
  593. return (n != 0) and (n & (n - 1) == 0)
  594. def _background_thread_fetch_current_state(
  595. serializer: SerializerBase, pipe: mp.connection.Connection, get_current_state_ref: weakref.WeakMethod
  596. ):
  597. """
  598. Executed in the host process as a background thread. Fetches the averager state when asked by peers.
  599. :param serializer: a serializer with which to convert metadata into bytes
  600. :param pipe: DecentralizedAverager's control pipe (from host process side)
  601. :param get_current_state_ref: a WeakMethod wrapped around DecentralizedAverager.get_current_state (instance-bound)
  602. """
  603. while True:
  604. try:
  605. trigger, future = pipe.recv()
  606. except BaseException as e:
  607. logger.debug(f"Averager background thread finished: {repr(e)}")
  608. break
  609. if trigger == "_SHUTDOWN":
  610. break
  611. assert trigger == "_TRIGGER_GET_CURRENT_STATE"
  612. try:
  613. get_current_state = get_current_state_ref()
  614. if get_current_state is None:
  615. break
  616. state = get_current_state()
  617. assert 0 < len(state) <= 3
  618. if len(state) != 3:
  619. state = tuple(state + (None,) * (3 - len(state)))
  620. state_metadata, state_tensors, tensor_infos = state
  621. del get_current_state
  622. state_metadata = serializer.dumps(state_metadata)
  623. state_tensors = tuple(
  624. tensor.cpu().detach().requires_grad_(tensor.requires_grad) for tensor in state_tensors
  625. )
  626. # note: we cast tensors to CPU on host side to avoid initializing cuda in the guest process
  627. future.set_result((state_metadata, state_tensors, tensor_infos))
  628. except BaseException as e:
  629. future.set_exception(e)
  630. logger.warning(e)
  631. continue
  632. def compute_schema_hash(tensors: Sequence[torch.Tensor]) -> bytes:
  633. """A hash that describes follower's tensor shapes, dtypes, devices, but not the actual values"""
  634. schema_dicts = [
  635. {
  636. field_name: str(field_value)
  637. for field_name, field_value in asdict(TensorDescriptor.from_tensor(tensor)).items()
  638. }
  639. for tensor in tensors
  640. ]
  641. return DHTID.generate(source=schema_dicts).to_bytes()