test_moe.py 7.8 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173
  1. import grpc
  2. import numpy as np
  3. import pytest
  4. import torch
  5. import hivemind
  6. from hivemind.client.expert import DUMMY
  7. from hivemind import background_server
  8. @pytest.mark.forked
  9. def test_moe():
  10. all_expert_uids = [f'ffn.{np.random.randint(0, 3)}.{np.random.randint(0, 3)}.{np.random.randint(0, 3)}'
  11. for _ in range(20)]
  12. with background_server(expert_uids=all_expert_uids, device='cpu', expert_cls='ffn',
  13. num_handlers=1, hidden_dim=16) as (server_endpoint, dht_endpoint):
  14. dht = hivemind.DHT(start=True, expiration=999, initial_peers=[dht_endpoint])
  15. dmoe = hivemind.RemoteMixtureOfExperts(
  16. in_features=16, grid_size=(32, 32, 32), dht=dht, k_best=3, uid_prefix='ffn.')
  17. for i in range(10):
  18. out = dmoe(torch.randn(10, 16))
  19. out.sum().backward()
  20. @pytest.mark.forked
  21. def test_call_many():
  22. k_min = 1
  23. timeout_after_k_min = None
  24. backward_k_min = 1
  25. forward_timeout = None
  26. backward_timeout = None
  27. rtol = 1e-3
  28. atol = 1e-5
  29. with background_server(num_experts=5, device='cpu', expert_cls='ffn', num_handlers=8, hidden_dim=64,
  30. optim_cls=None, no_dht=True) as (server_endpoint, dht_endpoint):
  31. inputs = torch.randn(4, 64, requires_grad=True)
  32. inputs_clone = inputs.clone().detach().requires_grad_(True)
  33. e0, e1, e2, e3, e4 = [hivemind.RemoteExpert(f'expert.{i}', server_endpoint) for i in range(5)]
  34. e5 = hivemind.RemoteExpert(f'thisshouldnotexist', '127.0.0.1:80')
  35. mask, expert_outputs = hivemind.client.moe._RemoteCallMany.apply(
  36. DUMMY, [[e0, e1, e2], [e2, e4], [e1, e5, e3], []],
  37. k_min, backward_k_min, timeout_after_k_min, forward_timeout, backward_timeout, e1.info, inputs
  38. )
  39. assert mask.shape == (4, 3)
  40. assert expert_outputs.shape == (4, 3, 64)
  41. assert np.all(mask.data.numpy() == np.array([[True, True, True],
  42. [True, True, False],
  43. [True, False, True],
  44. [False, False, False]])), f"Incorrect mask, {mask}"
  45. reference_outputs = torch.zeros_like(expert_outputs)
  46. reference_outputs[0, 0] = e0(inputs_clone[0:1])
  47. reference_outputs[0, 1] = e1(inputs_clone[0:1])
  48. reference_outputs[0, 2] = e2(inputs_clone[0:1])
  49. reference_outputs[1, 0] = e2(inputs_clone[1:2])
  50. reference_outputs[1, 1] = e4(inputs_clone[1:2])
  51. reference_outputs[2, 0] = e1(inputs_clone[2:3])
  52. reference_outputs[2, 2] = e3(inputs_clone[2:3])
  53. assert torch.allclose(expert_outputs, reference_outputs, rtol, atol)
  54. proj = torch.randn(4, 64)
  55. loss = (expert_outputs[(0, 1, 1, 2), (0, 2, 1, 0)] * proj).sum()
  56. loss.backward()
  57. our_grad = inputs.grad.data.cpu().clone()
  58. reference_loss = (reference_outputs[(0, 1, 1, 2), (0, 2, 1, 0)] * proj).sum()
  59. reference_loss.backward()
  60. reference_grad = inputs_clone.grad.data.cpu().clone()
  61. assert torch.allclose(our_grad, reference_grad, rtol, atol)
  62. @pytest.mark.forked
  63. def test_remote_module_call():
  64. with background_server(num_experts=1, device='cpu', expert_cls='ffn', num_handlers=1, hidden_dim=1024,
  65. optim_cls=None, no_dht=True) as (server_endpoint, dht_endpoint):
  66. real_expert = hivemind.RemoteExpert('expert.0', server_endpoint)
  67. fake_expert = hivemind.RemoteExpert('oiasfjiasjf', server_endpoint)
  68. out1 = real_expert(torch.randn(1, 1024))
  69. assert out1.shape == (1, 1024)
  70. dummy_x = torch.randn(3, 1024, requires_grad=True)
  71. out3 = real_expert(dummy_x)
  72. assert out3.shape == (3, 1024)
  73. out3_again = real_expert(dummy_x[1:])
  74. assert torch.allclose(out3_again, out3[1:], atol=1e-6, rtol=0)
  75. out3_again.norm().backward()
  76. assert dummy_x.grad is not None and dummy_x.grad.norm() > 0
  77. with pytest.raises(grpc.RpcError):
  78. real_expert(torch.randn(3, 11))
  79. with pytest.raises(grpc.RpcError):
  80. fake_expert(dummy_x)
  81. @pytest.mark.forked
  82. def test_beam_search_correctness():
  83. all_expert_uids = [f'ffn.{5 + i}.{10 + j}.{15 + k}' for i in range(10) for j in range(10) for k in range(10)]
  84. dht = hivemind.DHT(start=True, expiration=999)
  85. assert all(dht.declare_experts(all_expert_uids, endpoint='fake-endpoint'))
  86. dmoe = hivemind.RemoteMixtureOfExperts(
  87. in_features=32, grid_size=(32, 32, 32), dht=dht, k_best=4, uid_prefix='ffn.')
  88. for i in range(25):
  89. input = torch.randn(32)
  90. grid_scores = dmoe.proj(input).split_with_sizes(dmoe.grid_size, dim=-1)
  91. chosen_experts = dht.find_best_experts(dmoe.uid_prefix, [tensor.detach().numpy() for tensor in grid_scores],
  92. beam_size=dmoe.k_best)
  93. chosen_scores = dmoe.compute_expert_scores([dim_scores[None] for dim_scores in grid_scores],
  94. [chosen_experts])[0]
  95. our_best_scores = list(chosen_scores.cpu().detach().numpy())
  96. # reference: independently find :beam_size: best experts with exhaustive search
  97. all_scores = dmoe.compute_expert_scores([dim_scores.unsqueeze(0) for dim_scores in grid_scores],
  98. [[hivemind.RemoteExpert(uid, '') for uid in all_expert_uids]])[0]
  99. true_best_scores = sorted(all_scores.cpu().detach().numpy(), reverse=True)[:len(chosen_experts)]
  100. assert np.allclose(true_best_scores, our_best_scores)
  101. @pytest.mark.forked
  102. def test_determinism():
  103. rtol = 0
  104. atol = 1e-5
  105. xx = torch.randn(32, 1024, requires_grad=True)
  106. mask = torch.randint(0, 1, (32, 1024))
  107. with background_server(num_experts=1, device='cpu', expert_cls='det_dropout', num_handlers=1,
  108. optim_cls=None, no_dht=True) as (server_endpoint, dht_endpoint):
  109. expert = hivemind.RemoteExpert(uid=f'expert.0', endpoint=server_endpoint)
  110. out = expert(xx, mask)
  111. out_rerun = expert(xx, mask)
  112. grad, = torch.autograd.grad(out.sum(), xx, retain_graph=True)
  113. grad_rerun, = torch.autograd.grad(out_rerun.sum(), xx, retain_graph=True)
  114. assert torch.allclose(out, out_rerun, rtol, atol), "Dropout layer outputs are non-deterministic."
  115. assert torch.allclose(grad, grad_rerun, rtol, atol), "Gradients are non-deterministic."
  116. @pytest.mark.forked
  117. def test_compute_expert_scores():
  118. try:
  119. dht = hivemind.DHT(start=True)
  120. moe = hivemind.client.moe.RemoteMixtureOfExperts(
  121. dht=dht, in_features=1024, grid_size=(40,), k_best=4, k_min=1, timeout_after_k_min=1,
  122. uid_prefix='expert.')
  123. gx, gy = torch.randn(4, 5, requires_grad=True), torch.randn(4, 3, requires_grad=True)
  124. ii = [[4, 0, 2], [3, 1, 1, 1, 3], [0], [3, 2]]
  125. jj = [[2, 2, 1], [0, 1, 2, 0, 1], [0], [1, 2]]
  126. batch_experts = [
  127. [hivemind.RemoteExpert(uid=f'expert.{ii[batch_i][expert_i]}.{jj[batch_i][expert_i]}', endpoint="[::]:1337")
  128. for expert_i in range(len(ii[batch_i]))]
  129. for batch_i in range(len(ii))
  130. ] # note: these experts do not exists on server, we use them only to test moe compute_expert_scores
  131. logits = moe.compute_expert_scores([gx, gy], batch_experts)
  132. torch.softmax(logits, dim=-1).norm(dim=-1).mean().backward()
  133. assert gx.grad.norm().item() > 0 and gy.grad.norm().item(), "compute_expert_scores didn't backprop"
  134. for batch_i in range(len(ii)):
  135. for expert_i in range(len(ii[batch_i])):
  136. assert torch.allclose(logits[batch_i, expert_i],
  137. gx[batch_i, ii[batch_i][expert_i]] + gy[batch_i, jj[batch_i][expert_i]]), \
  138. "compute_expert_scores returned incorrect score"
  139. finally:
  140. dht.shutdown()