Spaces:
Running
Running
File size: 27,735 Bytes
85e3d20 |
1 2 3 4 5 6 7 8 9 10 11 12 13 14 15 16 17 18 19 20 21 22 23 24 25 26 27 28 29 30 31 32 33 34 35 36 37 38 39 40 41 42 43 44 45 46 47 48 49 50 51 52 53 54 55 56 57 58 59 60 61 62 63 64 65 66 67 68 69 70 71 72 73 74 75 76 77 78 79 80 81 82 83 84 85 86 87 88 89 90 91 92 93 94 95 96 97 98 99 100 101 102 103 104 105 106 107 108 109 110 111 112 113 114 115 116 117 118 119 120 121 122 123 124 125 126 127 128 129 130 131 132 133 134 135 136 137 138 139 140 141 142 143 144 145 146 147 148 149 150 151 152 153 154 155 156 157 158 159 160 161 162 163 164 165 166 167 168 169 170 171 172 173 174 175 176 177 178 179 180 181 182 183 184 185 186 187 188 189 190 191 192 193 194 195 196 197 198 199 200 201 202 203 204 205 206 207 208 209 210 211 212 213 214 215 216 217 218 219 220 221 222 223 224 225 226 227 228 229 230 231 232 233 234 235 236 237 238 239 240 241 242 243 244 245 246 247 248 249 250 251 252 253 254 255 256 257 258 259 260 261 262 263 264 265 266 267 268 269 270 271 272 273 274 275 276 277 278 279 280 281 282 283 284 285 286 287 288 289 290 291 292 293 294 295 296 297 298 299 300 301 302 303 304 305 306 307 308 309 310 311 312 313 314 315 316 317 318 319 320 321 322 323 324 325 326 327 328 329 330 331 332 333 334 335 336 337 338 339 340 341 342 343 344 345 346 347 348 349 350 351 352 353 354 355 356 357 358 359 360 361 362 363 364 365 366 367 368 369 370 371 372 373 374 375 376 377 378 379 380 381 382 383 384 385 386 387 388 389 390 391 392 393 394 395 396 397 398 399 400 401 402 403 404 405 406 407 408 409 410 411 412 413 414 415 416 417 418 419 420 421 422 423 424 425 426 427 428 429 430 431 432 433 434 435 436 437 438 439 440 441 442 443 444 445 446 447 448 449 450 451 452 453 454 455 456 457 458 459 460 461 462 463 464 465 466 467 468 469 470 471 472 473 474 475 476 477 478 479 480 481 482 483 484 485 486 487 488 489 490 491 492 493 494 495 496 497 498 499 500 501 502 503 504 505 506 507 508 509 510 511 512 513 514 515 516 517 518 519 520 521 522 523 524 525 526 527 528 529 530 531 532 533 534 535 536 537 538 539 540 541 542 543 544 545 546 547 548 549 550 551 552 553 554 555 556 557 558 559 560 561 562 563 564 565 566 567 568 569 570 571 572 573 574 575 576 577 578 579 580 581 582 583 584 585 586 587 588 589 590 591 592 593 594 595 596 597 598 599 600 601 602 603 604 605 606 607 608 609 610 611 612 613 614 615 616 617 618 619 620 621 622 623 624 625 626 627 628 629 630 631 632 633 634 635 636 637 638 639 640 641 642 643 644 645 646 647 648 649 650 651 652 653 654 655 656 657 658 659 660 661 662 663 664 665 666 667 668 669 670 671 672 673 674 675 676 677 678 679 680 681 682 683 684 685 686 687 688 689 690 691 692 693 694 695 696 697 698 699 700 701 702 703 704 705 706 707 708 709 710 711 712 713 714 715 716 717 718 719 720 721 722 723 724 725 726 727 728 729 730 731 732 733 734 735 736 737 738 739 740 741 742 743 744 745 746 747 748 749 750 751 752 753 754 755 756 757 758 759 760 761 762 763 764 765 766 767 768 769 770 771 772 773 774 775 776 777 778 779 780 781 782 783 784 785 786 787 788 789 790 791 792 793 794 795 796 797 798 799 800 801 802 803 804 805 806 807 808 809 810 811 812 813 814 815 816 817 818 819 820 821 822 823 824 825 826 827 828 829 830 831 832 833 834 835 836 837 838 839 840 841 842 843 844 845 846 847 848 849 850 851 852 853 854 855 856 857 |
# Copyright 2021 DeepMind Technologies Limited. All Rights Reserved.
#
# Licensed under the Apache License, Version 2.0 (the "License");
# you may not use this file except in compliance with the License.
# You may obtain a copy of the License at
#
# http://www.apache.org/licenses/LICENSE-2.0
#
# Unless required by applicable law or agreed to in writing, software
# distributed under the License is distributed on an "AS IS" BASIS,
# WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
# See the License for the specific language governing permissions and
# limitations under the License.
# ==============================================================================
"""JAX implementation of baseline processor networks."""
import abc
from typing import Any, Callable, List, Optional, Tuple
import chex
import haiku as hk
import jax
import jax.numpy as jnp
import numpy as np
_Array = chex.Array
_Fn = Callable[..., Any]
BIG_NUMBER = 1e6
PROCESSOR_TAG = 'clrs_processor'
class Processor(hk.Module):
"""Processor abstract base class."""
def __init__(self, name: str):
if not name.endswith(PROCESSOR_TAG):
name = name + '_' + PROCESSOR_TAG
super().__init__(name=name)
@abc.abstractmethod
def __call__(
self,
node_fts: _Array,
edge_fts: _Array,
graph_fts: _Array,
adj_mat: _Array,
hidden: _Array,
**kwargs,
) -> Tuple[_Array, Optional[_Array]]:
"""Processor inference step.
Args:
node_fts: Node features.
edge_fts: Edge features.
graph_fts: Graph features.
adj_mat: Graph adjacency matrix.
hidden: Hidden features.
**kwargs: Extra kwargs.
Returns:
Output of processor inference step as a 2-tuple of (node, edge)
embeddings. The edge embeddings can be None.
"""
pass
@property
def inf_bias(self):
return False
@property
def inf_bias_edge(self):
return False
class GAT(Processor):
"""Graph Attention Network (Velickovic et al., ICLR 2018)."""
def __init__(
self,
out_size: int,
nb_heads: int,
activation: Optional[_Fn] = jax.nn.relu,
residual: bool = True,
use_ln: bool = False,
name: str = 'gat_aggr',
):
super().__init__(name=name)
self.out_size = out_size
self.nb_heads = nb_heads
if out_size % nb_heads != 0:
raise ValueError('The number of attention heads must divide the width!')
self.head_size = out_size // nb_heads
self.activation = activation
self.residual = residual
self.use_ln = use_ln
def __call__( # pytype: disable=signature-mismatch # numpy-scalars
self,
node_fts: _Array,
edge_fts: _Array,
graph_fts: _Array,
adj_mat: _Array,
hidden: _Array,
**unused_kwargs,
) -> _Array:
"""GAT inference step."""
b, n, _ = node_fts.shape
assert edge_fts.shape[:-1] == (b, n, n)
assert graph_fts.shape[:-1] == (b,)
assert adj_mat.shape == (b, n, n)
z = jnp.concatenate([node_fts, hidden], axis=-1)
m = hk.Linear(self.out_size)
skip = hk.Linear(self.out_size)
bias_mat = (adj_mat - 1.0) * 1e9
bias_mat = jnp.tile(bias_mat[..., None],
(1, 1, 1, self.nb_heads)) # [B, N, N, H]
bias_mat = jnp.transpose(bias_mat, (0, 3, 1, 2)) # [B, H, N, N]
a_1 = hk.Linear(self.nb_heads)
a_2 = hk.Linear(self.nb_heads)
a_e = hk.Linear(self.nb_heads)
a_g = hk.Linear(self.nb_heads)
values = m(z) # [B, N, H*F]
values = jnp.reshape(
values,
values.shape[:-1] + (self.nb_heads, self.head_size)) # [B, N, H, F]
values = jnp.transpose(values, (0, 2, 1, 3)) # [B, H, N, F]
att_1 = jnp.expand_dims(a_1(z), axis=-1)
att_2 = jnp.expand_dims(a_2(z), axis=-1)
att_e = a_e(edge_fts)
att_g = jnp.expand_dims(a_g(graph_fts), axis=-1)
logits = (
jnp.transpose(att_1, (0, 2, 1, 3)) + # + [B, H, N, 1]
jnp.transpose(att_2, (0, 2, 3, 1)) + # + [B, H, 1, N]
jnp.transpose(att_e, (0, 3, 1, 2)) + # + [B, H, N, N]
jnp.expand_dims(att_g, axis=-1) # + [B, H, 1, 1]
) # = [B, H, N, N]
coefs = jax.nn.softmax(jax.nn.leaky_relu(logits) + bias_mat, axis=-1)
ret = jnp.matmul(coefs, values) # [B, H, N, F]
ret = jnp.transpose(ret, (0, 2, 1, 3)) # [B, N, H, F]
ret = jnp.reshape(ret, ret.shape[:-2] + (self.out_size,)) # [B, N, H*F]
if self.residual:
ret += skip(z)
if self.activation is not None:
ret = self.activation(ret)
if self.use_ln:
ln = hk.LayerNorm(axis=-1, create_scale=True, create_offset=True)
ret = ln(ret)
return ret, None # pytype: disable=bad-return-type # numpy-scalars
class GATFull(GAT):
"""Graph Attention Network with full adjacency matrix."""
def __call__(self, node_fts: _Array, edge_fts: _Array, graph_fts: _Array,
adj_mat: _Array, hidden: _Array, **unused_kwargs) -> _Array:
adj_mat = jnp.ones_like(adj_mat)
return super().__call__(node_fts, edge_fts, graph_fts, adj_mat, hidden)
class GATv2(Processor):
"""Graph Attention Network v2 (Brody et al., ICLR 2022)."""
def __init__(
self,
out_size: int,
nb_heads: int,
mid_size: Optional[int] = None,
activation: Optional[_Fn] = jax.nn.relu,
residual: bool = True,
use_ln: bool = False,
name: str = 'gatv2_aggr',
):
super().__init__(name=name)
if mid_size is None:
self.mid_size = out_size
else:
self.mid_size = mid_size
self.out_size = out_size
self.nb_heads = nb_heads
if out_size % nb_heads != 0:
raise ValueError('The number of attention heads must divide the width!')
self.head_size = out_size // nb_heads
if self.mid_size % nb_heads != 0:
raise ValueError('The number of attention heads must divide the message!')
self.mid_head_size = self.mid_size // nb_heads
self.activation = activation
self.residual = residual
self.use_ln = use_ln
def __call__( # pytype: disable=signature-mismatch # numpy-scalars
self,
node_fts: _Array,
edge_fts: _Array,
graph_fts: _Array,
adj_mat: _Array,
hidden: _Array,
**unused_kwargs,
) -> _Array:
"""GATv2 inference step."""
b, n, _ = node_fts.shape
assert edge_fts.shape[:-1] == (b, n, n)
assert graph_fts.shape[:-1] == (b,)
assert adj_mat.shape == (b, n, n)
z = jnp.concatenate([node_fts, hidden], axis=-1)
m = hk.Linear(self.out_size)
skip = hk.Linear(self.out_size)
bias_mat = (adj_mat - 1.0) * 1e9
bias_mat = jnp.tile(bias_mat[..., None],
(1, 1, 1, self.nb_heads)) # [B, N, N, H]
bias_mat = jnp.transpose(bias_mat, (0, 3, 1, 2)) # [B, H, N, N]
w_1 = hk.Linear(self.mid_size)
w_2 = hk.Linear(self.mid_size)
w_e = hk.Linear(self.mid_size)
w_g = hk.Linear(self.mid_size)
a_heads = []
for _ in range(self.nb_heads):
a_heads.append(hk.Linear(1))
values = m(z) # [B, N, H*F]
values = jnp.reshape(
values,
values.shape[:-1] + (self.nb_heads, self.head_size)) # [B, N, H, F]
values = jnp.transpose(values, (0, 2, 1, 3)) # [B, H, N, F]
pre_att_1 = w_1(z)
pre_att_2 = w_2(z)
pre_att_e = w_e(edge_fts)
pre_att_g = w_g(graph_fts)
pre_att = (
jnp.expand_dims(pre_att_1, axis=1) + # + [B, 1, N, H*F]
jnp.expand_dims(pre_att_2, axis=2) + # + [B, N, 1, H*F]
pre_att_e + # + [B, N, N, H*F]
jnp.expand_dims(pre_att_g, axis=(1, 2)) # + [B, 1, 1, H*F]
) # = [B, N, N, H*F]
pre_att = jnp.reshape(
pre_att,
pre_att.shape[:-1] + (self.nb_heads, self.mid_head_size)
) # [B, N, N, H, F]
pre_att = jnp.transpose(pre_att, (0, 3, 1, 2, 4)) # [B, H, N, N, F]
# This part is not very efficient, but we agree to keep it this way to
# enhance readability, assuming `nb_heads` will not be large.
logit_heads = []
for head in range(self.nb_heads):
logit_heads.append(
jnp.squeeze(
a_heads[head](jax.nn.leaky_relu(pre_att[:, head])),
axis=-1)
) # [B, N, N]
logits = jnp.stack(logit_heads, axis=1) # [B, H, N, N]
coefs = jax.nn.softmax(logits + bias_mat, axis=-1)
ret = jnp.matmul(coefs, values) # [B, H, N, F]
ret = jnp.transpose(ret, (0, 2, 1, 3)) # [B, N, H, F]
ret = jnp.reshape(ret, ret.shape[:-2] + (self.out_size,)) # [B, N, H*F]
if self.residual:
ret += skip(z)
if self.activation is not None:
ret = self.activation(ret)
if self.use_ln:
ln = hk.LayerNorm(axis=-1, create_scale=True, create_offset=True)
ret = ln(ret)
return ret, None # pytype: disable=bad-return-type # numpy-scalars
class GATv2Full(GATv2):
"""Graph Attention Network v2 with full adjacency matrix."""
def __call__(self, node_fts: _Array, edge_fts: _Array, graph_fts: _Array,
adj_mat: _Array, hidden: _Array, **unused_kwargs) -> _Array:
adj_mat = jnp.ones_like(adj_mat)
return super().__call__(node_fts, edge_fts, graph_fts, adj_mat, hidden)
def get_triplet_msgs(z, edge_fts, graph_fts, nb_triplet_fts):
"""Triplet messages, as done by Dudzik and Velickovic (2022)."""
t_1 = hk.Linear(nb_triplet_fts)
t_2 = hk.Linear(nb_triplet_fts)
t_3 = hk.Linear(nb_triplet_fts)
t_e_1 = hk.Linear(nb_triplet_fts)
t_e_2 = hk.Linear(nb_triplet_fts)
t_e_3 = hk.Linear(nb_triplet_fts)
t_g = hk.Linear(nb_triplet_fts)
tri_1 = t_1(z)
tri_2 = t_2(z)
tri_3 = t_3(z)
tri_e_1 = t_e_1(edge_fts)
tri_e_2 = t_e_2(edge_fts)
tri_e_3 = t_e_3(edge_fts)
tri_g = t_g(graph_fts)
return (
jnp.expand_dims(tri_1, axis=(2, 3)) + # (B, N, 1, 1, H)
jnp.expand_dims(tri_2, axis=(1, 3)) + # + (B, 1, N, 1, H)
jnp.expand_dims(tri_3, axis=(1, 2)) + # + (B, 1, 1, N, H)
jnp.expand_dims(tri_e_1, axis=3) + # + (B, N, N, 1, H)
jnp.expand_dims(tri_e_2, axis=2) + # + (B, N, 1, N, H)
jnp.expand_dims(tri_e_3, axis=1) + # + (B, 1, N, N, H)
jnp.expand_dims(tri_g, axis=(1, 2, 3)) # + (B, 1, 1, 1, H)
) # = (B, N, N, N, H)
class PGN(Processor):
"""Pointer Graph Networks (Veličković et al., NeurIPS 2020)."""
def __init__(
self,
out_size: int,
mid_size: Optional[int] = None,
mid_act: Optional[_Fn] = None,
activation: Optional[_Fn] = jax.nn.relu,
reduction: _Fn = jnp.max,
msgs_mlp_sizes: Optional[List[int]] = None,
use_ln: bool = False,
use_triplets: bool = False,
nb_triplet_fts: int = 8,
gated: bool = False,
name: str = 'mpnn_aggr',
):
super().__init__(name=name)
if mid_size is None:
self.mid_size = out_size
else:
self.mid_size = mid_size
self.out_size = out_size
self.mid_act = mid_act
self.activation = activation
self.reduction = reduction
self._msgs_mlp_sizes = msgs_mlp_sizes
self.use_ln = use_ln
self.use_triplets = use_triplets
self.nb_triplet_fts = nb_triplet_fts
self.gated = gated
def __call__( # pytype: disable=signature-mismatch # numpy-scalars
self,
node_fts: _Array,
edge_fts: _Array,
graph_fts: _Array,
adj_mat: _Array,
hidden: _Array,
**unused_kwargs,
) -> _Array:
"""MPNN inference step."""
b, n, _ = node_fts.shape
assert edge_fts.shape[:-1] == (b, n, n)
assert graph_fts.shape[:-1] == (b,)
assert adj_mat.shape == (b, n, n)
z = jnp.concatenate([node_fts, hidden], axis=-1)
m_1 = hk.Linear(self.mid_size)
m_2 = hk.Linear(self.mid_size)
m_e = hk.Linear(self.mid_size)
m_g = hk.Linear(self.mid_size)
o1 = hk.Linear(self.out_size)
o2 = hk.Linear(self.out_size)
msg_1 = m_1(z)
msg_2 = m_2(z)
msg_e = m_e(edge_fts)
msg_g = m_g(graph_fts)
tri_msgs = None
if self.use_triplets:
# Triplet messages, as done by Dudzik and Velickovic (2022)
triplets = get_triplet_msgs(z, edge_fts, graph_fts, self.nb_triplet_fts)
o3 = hk.Linear(self.out_size)
tri_msgs = o3(jnp.max(triplets, axis=1)) # (B, N, N, H)
if self.activation is not None:
tri_msgs = self.activation(tri_msgs)
msgs = (
jnp.expand_dims(msg_1, axis=1) + jnp.expand_dims(msg_2, axis=2) +
msg_e + jnp.expand_dims(msg_g, axis=(1, 2)))
if self._msgs_mlp_sizes is not None:
msgs = hk.nets.MLP(self._msgs_mlp_sizes)(jax.nn.relu(msgs))
if self.mid_act is not None:
msgs = self.mid_act(msgs)
if self.reduction == jnp.mean:
msgs = jnp.sum(msgs * jnp.expand_dims(adj_mat, -1), axis=1)
msgs = msgs / jnp.sum(adj_mat, axis=-1, keepdims=True)
elif self.reduction == jnp.max:
maxarg = jnp.where(jnp.expand_dims(adj_mat, -1),
msgs,
-BIG_NUMBER)
msgs = jnp.max(maxarg, axis=1)
else:
msgs = self.reduction(msgs * jnp.expand_dims(adj_mat, -1), axis=1)
h_1 = o1(z)
h_2 = o2(msgs)
ret = h_1 + h_2
if self.activation is not None:
ret = self.activation(ret)
if self.use_ln:
ln = hk.LayerNorm(axis=-1, create_scale=True, create_offset=True)
ret = ln(ret)
if self.gated:
gate1 = hk.Linear(self.out_size)
gate2 = hk.Linear(self.out_size)
gate3 = hk.Linear(self.out_size, b_init=hk.initializers.Constant(-3))
gate = jax.nn.sigmoid(gate3(jax.nn.relu(gate1(z) + gate2(msgs))))
ret = ret * gate + hidden * (1-gate)
return ret, tri_msgs # pytype: disable=bad-return-type # numpy-scalars
class DeepSets(PGN):
"""Deep Sets (Zaheer et al., NeurIPS 2017)."""
def __call__(self, node_fts: _Array, edge_fts: _Array, graph_fts: _Array,
adj_mat: _Array, hidden: _Array, **unused_kwargs) -> _Array:
assert adj_mat.ndim == 3
adj_mat = jnp.ones_like(adj_mat) * jnp.eye(adj_mat.shape[-1])
return super().__call__(node_fts, edge_fts, graph_fts, adj_mat, hidden)
class MPNN(PGN):
"""Message-Passing Neural Network (Gilmer et al., ICML 2017)."""
def __call__(self, node_fts: _Array, edge_fts: _Array, graph_fts: _Array,
adj_mat: _Array, hidden: _Array, **unused_kwargs) -> _Array:
adj_mat = jnp.ones_like(adj_mat)
return super().__call__(node_fts, edge_fts, graph_fts, adj_mat, hidden)
class PGNMask(PGN):
"""Masked Pointer Graph Networks (Veličković et al., NeurIPS 2020)."""
@property
def inf_bias(self):
return True
@property
def inf_bias_edge(self):
return True
class MemNetMasked(Processor):
"""Implementation of End-to-End Memory Networks.
Inspired by the description in https://arxiv.org/abs/1503.08895.
"""
def __init__(
self,
vocab_size: int,
sentence_size: int,
linear_output_size: int,
embedding_size: int = 16,
memory_size: Optional[int] = 128,
num_hops: int = 1,
nonlin: Callable[[Any], Any] = jax.nn.relu,
apply_embeddings: bool = True,
init_func: hk.initializers.Initializer = jnp.zeros,
use_ln: bool = False,
name: str = 'memnet') -> None:
"""Constructor.
Args:
vocab_size: the number of words in the dictionary (each story, query and
answer come contain symbols coming from this dictionary).
sentence_size: the dimensionality of each memory.
linear_output_size: the dimensionality of the output of the last layer
of the model.
embedding_size: the dimensionality of the latent space to where all
memories are projected.
memory_size: the number of memories provided.
num_hops: the number of layers in the model.
nonlin: non-linear transformation applied at the end of each layer.
apply_embeddings: flag whether to aply embeddings.
init_func: initialization function for the biases.
use_ln: whether to use layer normalisation in the model.
name: the name of the model.
"""
super().__init__(name=name)
self._vocab_size = vocab_size
self._embedding_size = embedding_size
self._sentence_size = sentence_size
self._memory_size = memory_size
self._linear_output_size = linear_output_size
self._num_hops = num_hops
self._nonlin = nonlin
self._apply_embeddings = apply_embeddings
self._init_func = init_func
self._use_ln = use_ln
# Encoding part: i.e. "I" of the paper.
self._encodings = _position_encoding(sentence_size, embedding_size)
def __call__( # pytype: disable=signature-mismatch # numpy-scalars
self,
node_fts: _Array,
edge_fts: _Array,
graph_fts: _Array,
adj_mat: _Array,
hidden: _Array,
**unused_kwargs,
) -> _Array:
"""MemNet inference step."""
del hidden
node_and_graph_fts = jnp.concatenate([node_fts, graph_fts[:, None]],
axis=1)
edge_fts_padded = jnp.pad(edge_fts * adj_mat[..., None],
((0, 0), (0, 1), (0, 1), (0, 0)))
nxt_hidden = jax.vmap(self._apply, (1), 1)(node_and_graph_fts,
edge_fts_padded)
# Broadcast hidden state corresponding to graph features across the nodes.
nxt_hidden = nxt_hidden[:, :-1] + nxt_hidden[:, -1:]
return nxt_hidden, None # pytype: disable=bad-return-type # numpy-scalars
def _apply(self, queries: _Array, stories: _Array) -> _Array:
"""Apply Memory Network to the queries and stories.
Args:
queries: Tensor of shape [batch_size, sentence_size].
stories: Tensor of shape [batch_size, memory_size, sentence_size].
Returns:
Tensor of shape [batch_size, vocab_size].
"""
if self._apply_embeddings:
query_biases = hk.get_parameter(
'query_biases',
shape=[self._vocab_size - 1, self._embedding_size],
init=self._init_func)
stories_biases = hk.get_parameter(
'stories_biases',
shape=[self._vocab_size - 1, self._embedding_size],
init=self._init_func)
memory_biases = hk.get_parameter(
'memory_contents',
shape=[self._memory_size, self._embedding_size],
init=self._init_func)
output_biases = hk.get_parameter(
'output_biases',
shape=[self._vocab_size - 1, self._embedding_size],
init=self._init_func)
nil_word_slot = jnp.zeros([1, self._embedding_size])
# This is "A" in the paper.
if self._apply_embeddings:
stories_biases = jnp.concatenate([stories_biases, nil_word_slot], axis=0)
memory_embeddings = jnp.take(
stories_biases, stories.reshape([-1]).astype(jnp.int32),
axis=0).reshape(list(stories.shape) + [self._embedding_size])
memory_embeddings = jnp.pad(
memory_embeddings,
((0, 0), (0, self._memory_size - jnp.shape(memory_embeddings)[1]),
(0, 0), (0, 0)))
memory = jnp.sum(memory_embeddings * self._encodings, 2) + memory_biases
else:
memory = stories
# This is "B" in the paper. Also, when there are no queries (only
# sentences), then there these lines are substituted by
# query_embeddings = 0.1.
if self._apply_embeddings:
query_biases = jnp.concatenate([query_biases, nil_word_slot], axis=0)
query_embeddings = jnp.take(
query_biases, queries.reshape([-1]).astype(jnp.int32),
axis=0).reshape(list(queries.shape) + [self._embedding_size])
# This is "u" in the paper.
query_input_embedding = jnp.sum(query_embeddings * self._encodings, 1)
else:
query_input_embedding = queries
# This is "C" in the paper.
if self._apply_embeddings:
output_biases = jnp.concatenate([output_biases, nil_word_slot], axis=0)
output_embeddings = jnp.take(
output_biases, stories.reshape([-1]).astype(jnp.int32),
axis=0).reshape(list(stories.shape) + [self._embedding_size])
output_embeddings = jnp.pad(
output_embeddings,
((0, 0), (0, self._memory_size - jnp.shape(output_embeddings)[1]),
(0, 0), (0, 0)))
output = jnp.sum(output_embeddings * self._encodings, 2)
else:
output = stories
intermediate_linear = hk.Linear(self._embedding_size, with_bias=False)
# Output_linear is "H".
output_linear = hk.Linear(self._linear_output_size, with_bias=False)
for hop_number in range(self._num_hops):
query_input_embedding_transposed = jnp.transpose(
jnp.expand_dims(query_input_embedding, -1), [0, 2, 1])
# Calculate probabilities.
probs = jax.nn.softmax(
jnp.sum(memory * query_input_embedding_transposed, 2))
# Calculate output of the layer by multiplying by C.
transposed_probs = jnp.transpose(jnp.expand_dims(probs, -1), [0, 2, 1])
transposed_output_embeddings = jnp.transpose(output, [0, 2, 1])
# This is "o" in the paper.
layer_output = jnp.sum(transposed_output_embeddings * transposed_probs, 2)
# Finally the answer
if hop_number == self._num_hops - 1:
# Please note that in the TF version we apply the final linear layer
# in all hops and this results in shape mismatches.
output_layer = output_linear(query_input_embedding + layer_output)
else:
output_layer = intermediate_linear(query_input_embedding + layer_output)
query_input_embedding = output_layer
if self._nonlin:
output_layer = self._nonlin(output_layer)
# This linear here is "W".
ret = hk.Linear(self._vocab_size, with_bias=False)(output_layer)
if self._use_ln:
ln = hk.LayerNorm(axis=-1, create_scale=True, create_offset=True)
ret = ln(ret)
return ret
class MemNetFull(MemNetMasked):
"""Memory Networks with full adjacency matrix."""
def __call__(self, node_fts: _Array, edge_fts: _Array, graph_fts: _Array,
adj_mat: _Array, hidden: _Array, **unused_kwargs) -> _Array:
adj_mat = jnp.ones_like(adj_mat)
return super().__call__(node_fts, edge_fts, graph_fts, adj_mat, hidden)
ProcessorFactory = Callable[[int], Processor]
def get_processor_factory(kind: str,
use_ln: bool,
nb_triplet_fts: int,
nb_heads: Optional[int] = None) -> ProcessorFactory:
"""Returns a processor factory.
Args:
kind: One of the available types of processor.
use_ln: Whether the processor passes the output through a layernorm layer.
nb_triplet_fts: How many triplet features to compute.
nb_heads: Number of attention heads for GAT processors.
Returns:
A callable that takes an `out_size` parameter (equal to the hidden
dimension of the network) and returns a processor instance.
"""
def _factory(out_size: int):
if kind == 'deepsets':
processor = DeepSets(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=False,
nb_triplet_fts=0
)
elif kind == 'gat':
processor = GAT(
out_size=out_size,
nb_heads=nb_heads,
use_ln=use_ln,
)
elif kind == 'gat_full':
processor = GATFull(
out_size=out_size,
nb_heads=nb_heads,
use_ln=use_ln
)
elif kind == 'gatv2':
processor = GATv2(
out_size=out_size,
nb_heads=nb_heads,
use_ln=use_ln
)
elif kind == 'gatv2_full':
processor = GATv2Full(
out_size=out_size,
nb_heads=nb_heads,
use_ln=use_ln
)
elif kind == 'memnet_full':
processor = MemNetFull(
vocab_size=out_size,
sentence_size=out_size,
linear_output_size=out_size,
)
elif kind == 'memnet_masked':
processor = MemNetMasked(
vocab_size=out_size,
sentence_size=out_size,
linear_output_size=out_size,
)
elif kind == 'mpnn':
processor = MPNN(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=False,
nb_triplet_fts=0,
)
elif kind == 'pgn':
processor = PGN(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=False,
nb_triplet_fts=0,
)
elif kind == 'pgn_mask':
processor = PGNMask(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=False,
nb_triplet_fts=0,
)
elif kind == 'triplet_mpnn':
processor = MPNN(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=True,
nb_triplet_fts=nb_triplet_fts,
)
elif kind == 'triplet_pgn':
processor = PGN(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=True,
nb_triplet_fts=nb_triplet_fts,
)
elif kind == 'triplet_pgn_mask':
processor = PGNMask(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=True,
nb_triplet_fts=nb_triplet_fts,
)
elif kind == 'gpgn':
processor = PGN(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=False,
nb_triplet_fts=nb_triplet_fts,
gated=True,
)
elif kind == 'gpgn_mask':
processor = PGNMask(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=False,
nb_triplet_fts=nb_triplet_fts,
gated=True,
)
elif kind == 'gmpnn':
processor = MPNN(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=False,
nb_triplet_fts=nb_triplet_fts,
gated=True,
)
elif kind == 'triplet_gpgn':
processor = PGN(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=True,
nb_triplet_fts=nb_triplet_fts,
gated=True,
)
elif kind == 'triplet_gpgn_mask':
processor = PGNMask(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=True,
nb_triplet_fts=nb_triplet_fts,
gated=True,
)
elif kind == 'triplet_gmpnn':
processor = MPNN(
out_size=out_size,
msgs_mlp_sizes=[out_size, out_size],
use_ln=use_ln,
use_triplets=True,
nb_triplet_fts=nb_triplet_fts,
gated=True,
)
else:
raise ValueError('Unexpected processor kind ' + kind)
return processor
return _factory
def _position_encoding(sentence_size: int, embedding_size: int) -> np.ndarray:
"""Position Encoding described in section 4.1 [1]."""
encoding = np.ones((embedding_size, sentence_size), dtype=np.float32)
ls = sentence_size + 1
le = embedding_size + 1
for i in range(1, le):
for j in range(1, ls):
encoding[i - 1, j - 1] = (i - (le - 1) / 2) * (j - (ls - 1) / 2)
encoding = 1 + 4 * encoding / embedding_size / sentence_size
return np.transpose(encoding)
|