moe.py 10.0 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188
  1. import multiprocessing as mp
  2. import multiprocessing.pool
  3. from concurrent.futures import as_completed
  4. from typing import Tuple, List, Dict, Any
  5. import numpy as np
  6. import torch
  7. import torch.nn as nn
  8. from .expert import RemoteExpert
  9. from ..utils import nested_map, check_numpy, run_in_background
  10. class RemoteMixtureOfExperts(nn.Module):
  11. """
  12. A torch module that performs mixture of experts inference with a local gating function and multiple remote experts.
  13. Natively supports pytorch autograd.
  14. :note: By default, not all experts are guaranteed to perform forward pass. Moreover, not all of those who ran
  15. forward pass are guaranteed to perform backward pass. In the latter case, gradient will be averaged without
  16. the missing experts
  17. :param in_features: common input size for experts and gating function
  18. :param grid_size: tesseract dimensions that form expert uid (see below)
  19. :param uid_prefix: common prefix for all expert uids
  20. expert uid follows the pattern {uid_prefix}{0...grid_size[0]}.{0...grid_size[1]}...{0...grid_size[-1]}
  21. :param network: TesseractNetwork where the experts reside
  22. :param num_workers: number of threads for parallel network operation
  23. :param k_best: queries this many experts with highest scores
  24. :param k_min: makes sure at least this many experts returned output
  25. :param timeout_after_k_min: waits for this many seconds after k_min experts returned results.
  26. Any expert that didn't manage to return output after that delay is considered unavailable
  27. :param expert_padding: internal value used to denote "absent expert". Should not coincide with any expert uid.
  28. """
  29. def __init__(self, *, in_features, grid_size: Tuple[int], network, num_workers=None,
  30. k_best, k_min=1, timeout_after_k_min=1.0, uid_prefix='', expert_padding=None):
  31. super().__init__()
  32. self.network, self.grid_size = network, grid_size
  33. self.uid_prefix, self.expert_padding = uid_prefix, expert_padding
  34. self.k_best, self.k_min, self.timeout_after_k_min = k_best, k_min, timeout_after_k_min
  35. self.thread_pool = mp.pool.ThreadPool(num_workers or k_best * 2)
  36. self.proj = nn.Linear(in_features, sum(grid_size)) # jointly predict logits for all grid dimensions
  37. def forward(self, input: torch.Tensor, *args, **kwargs) -> Tuple[List[List[RemoteExpert]], torch.Tensor]:
  38. """
  39. Choose k best experts with beam search, then call chosen experts and average their outputs.
  40. :param batch: named tensors, each tensor has 0-th axis dedicated to batch (aka batch-first
  41. :returns: averaged predictions of all experts that delivered on time
  42. """
  43. assert len(input.shape) == 2
  44. # 1. compute scores and find most appropriate experts with beam search
  45. grid_scores = self.proj(input).split_with_sizes(self.grid_size, dim=-1)
  46. batch_experts = self.beam_search(grid_scores, self.k_best)
  47. # ^-- List[batch_size] of List[RemoteExpert] chosen for every input in batch
  48. # 2.1 call chosen experts (run them in background to save time)
  49. batch_outputs_async = [
  50. self.thread_pool.apply_async(self._run_experts,
  51. args=[chosen_experts, input[i: i + 1], *(tensor[i: i + 1] for tensor in args)],
  52. kwds={key: tensor[i: i + 1] for key, tensor in kwargs.items()})
  53. for i, chosen_experts in enumerate(batch_experts)
  54. ]
  55. # 2.2 compute *differentiable* logits for each expert
  56. batch_expert_logits = self._score_experts(grid_scores, batch_experts)
  57. # ^-- List[batch_size] of Dict[RemoteExpert, logit] before softmax for each active expert
  58. batch_outputs = []
  59. for output_async, expert_logits in zip(batch_outputs_async, batch_expert_logits):
  60. expert_outputs: Dict[RemoteExpert, Any] = output_async.get()
  61. flat_experts, flat_outputs = zip(*expert_outputs.items())
  62. # 3.1. normalize logits over only those experts that DID return output
  63. flat_logits = torch.stack([expert_logits[expert] for expert in flat_experts])
  64. flat_weights = torch.softmax(flat_logits, dim=-1)
  65. # 3.2. average each output across experts
  66. average_outputs = nested_map(
  67. lambda *tensors: sum(x * weight for x, weight in zip(tensors, flat_weights)), *flat_outputs)
  68. batch_outputs.append(average_outputs)
  69. # 4. concatenate mixture outputs from individual experts
  70. return nested_map(lambda *tensors: torch.cat(tensors, dim=0), *batch_outputs)
  71. def beam_search(self, grid_scores: List[torch.Tensor], k_best: int, **kwargs) -> List[List[RemoteExpert]]:
  72. """
  73. Find and return k best experts in the grid using (exact) beam search of the product space
  74. :param grid_scores: scores predicted for each dimension in the grid,
  75. :type grid_scores: a sequence of tensors of shape[batch_size, self.grid_size[i]]
  76. :param k_best: how many of the top experts participate in the computation
  77. :param kwargs: extra keyword parameters passed to self.network.first_k_active
  78. :returns: a list of *batch_size* lists that contain chosen experts for one sample each inner list contains \
  79. RemoteExpert instances for *up to* k_best experts
  80. """
  81. assert len(grid_scores) == len(self.grid_size)
  82. assert all(len(dim_scores.shape) == 2 for dim_scores in grid_scores)
  83. batch_size = len(grid_scores[0])
  84. beam = np.array([[self.uid_prefix]] * batch_size, dtype=object) # [batch_size, up_to_beam_size]
  85. scores = np.zeros([batch_size, 1], dtype=np.float64)
  86. delimeters = np.array(self.network.UID_DELIMETER)[None, None, None] # pre-compute numpy array for fast concat
  87. for dim_index, dim_scores in enumerate(grid_scores):
  88. dim_scores = check_numpy(dim_scores)
  89. assert dim_scores.shape[-1] == self.grid_size[dim_index]
  90. # create all possible successsors from current beam
  91. dim_indices = np.arange(dim_scores.shape[1]).astype(str)
  92. new_candidates = beam[:, :, None] + delimeters + dim_indices[None, None, :]
  93. new_candidates = new_candidates.reshape([batch_size, -1])
  94. new_scores = scores[:, :, None] + dim_scores[:, None, :]
  95. new_scores = new_scores.reshape([batch_size, -1])
  96. # select k best candidates according to scores but only those that are still active
  97. new_order = np.argsort(- new_scores, axis=-1)
  98. top_alive_lookups = [
  99. self.thread_pool.apply_async(self.network.first_k_active, args=(cands[order], k_best), kwds=kwargs)
  100. for cands, order in zip(new_candidates, new_order)]
  101. batch_cand_to_score = [
  102. dict(zip(cands, cand_scores)) for cands, cand_scores in zip(new_candidates, new_scores)]
  103. top_alive_prefixes = [result.get() for result in top_alive_lookups]
  104. top_alive_scores = [list(map(cand_to_score.get, top_cands))
  105. for cand_to_score, top_cands in zip(batch_cand_to_score, top_alive_prefixes)]
  106. # pad up to beam size
  107. beam = np.array([row + [self.expert_padding] * (k_best - len(row))
  108. for row in top_alive_prefixes], dtype='object')
  109. scores = np.array([row + [-float('inf')] * (k_best - len(row))
  110. for row in top_alive_scores], dtype='float32')
  111. unique_experts = self.network.get_experts(list(set(
  112. uid for row in beam for uid in row if uid != self.expert_padding)))
  113. unique_experts_by_uid = {expert.uid: expert for expert in unique_experts if expert != self.expert_padding}
  114. return [
  115. [unique_experts_by_uid[uid] for uid in row if uid in unique_experts_by_uid]
  116. for row in beam]
  117. def _run_experts(self, experts: List[RemoteExpert], *args, **kwargs) -> Dict[RemoteExpert, torch.Tensor]:
  118. future_to_expert = {run_in_background(expert, *args, **kwargs): expert for expert in experts}
  119. pending_futures = set(future_to_expert.keys())
  120. outputs = {} # {expert -> output}
  121. # await first k futures for as long as it takes
  122. for future in as_completed(list(pending_futures), timeout=None):
  123. pending_futures.remove(future)
  124. if not future.exception():
  125. outputs[future_to_expert.pop(future)] = future.result()
  126. if len(outputs) > self.k_min:
  127. break
  128. # await stragglers for at most self.timeout_after_k_min
  129. for future in as_completed(pending_futures, timeout=self.timeout_after_k_min):
  130. if not future.exception():
  131. outputs[future_to_expert.pop(future)] = future.result()
  132. return outputs
  133. def _score_experts(self, grid_scores: List[torch.Tensor],
  134. experts: List[List[RemoteExpert]]) -> List[Dict[RemoteExpert, torch.Tensor]]:
  135. flat_experts = [expert for row in experts for expert in row]
  136. flat_batch_indices = torch.tensor([i for i, row in enumerate(experts) for uid in range(len(row))])
  137. grid_indices = np.zeros([len(flat_experts), len(grid_scores)], dtype=np.int64)
  138. for i, expert in enumerate(flat_experts):
  139. expert_indices = expert.uid[len(self.uid_prefix) + len(self.network.UID_DELIMETER):]
  140. expert_indices = list(map(int, expert_indices.split(self.network.UID_DELIMETER)))
  141. grid_indices[i] = expert_indices
  142. scores_per_dim = [
  143. dim_scores[flat_batch_indices, dim_indices] if len(flat_batch_indices) else torch.zeros(0)
  144. for dim_scores, dim_indices in zip(grid_scores, grid_indices.T)]
  145. flat_scores = torch.sum(torch.stack(scores_per_dim, dim=0), dim=0)
  146. output_dicts = [dict() for _ in range(len(experts))]
  147. for batch_i, expert, score in zip(check_numpy(flat_batch_indices),
  148. flat_experts, flat_scores):
  149. output_dicts[batch_i][expert] = score
  150. return output_dicts