Source code for torch_geometric.nn.conv.ppf_conv

from typing import Optional, Callable, Union
from torch_geometric.typing import OptTensor, PairOptTensor, PairTensor, Adj

import torch
from torch import Tensor
from torch_sparse import SparseTensor, set_diag
from torch_geometric.utils import remove_self_loops, add_self_loops
from torch_geometric.nn.conv import MessagePassing

from ..inits import reset


def get_angle(v1: Tensor, v2: Tensor) -> Tensor:
    return torch.atan2(
        torch.cross(v1, v2, dim=1).norm(p=2, dim=1), (v1 * v2).sum(dim=1))


def point_pair_features(pos_i: Tensor, pos_j: Tensor, norm_i: Tensor,
                        norm_j: Tensor) -> Tensor:
    pseudo = pos_j - pos_i
    return torch.stack([
        pseudo.norm(p=2, dim=1),
        get_angle(norm_i, pseudo),
        get_angle(norm_j, pseudo),
        get_angle(norm_i, norm_j)
    ], dim=1)


[docs]class PPFConv(MessagePassing): r"""The PPFNet operator from the `"PPFNet: Global Context Aware Local Features for Robust 3D Point Matching" <https://arxiv.org/abs/1802.02669>`_ paper .. math:: \mathbf{x}^{\prime}_i = \gamma_{\mathbf{\Theta}} \left( \max_{j \in \mathcal{N}(i) \cup \{ i \}} h_{\mathbf{\Theta}} ( \mathbf{x}_j, \| \mathbf{d_{j,i}} \|, \angle(\mathbf{n}_i, \mathbf{d_{j,i}}), \angle(\mathbf{n}_j, \mathbf{d_{j,i}}), \angle(\mathbf{n}_i, \mathbf{n}_j) \right) where :math:`\gamma_{\mathbf{\Theta}}` and :math:`h_{\mathbf{\Theta}}` denote neural networks, *.i.e.* MLPs, which takes in node features and :class:`torch_geometric.transforms.PointPairFeatures`. Args: local_nn (torch.nn.Module, optional): A neural network :math:`h_{\mathbf{\Theta}}` that maps node features :obj:`x` and relative spatial coordinates :obj:`pos_j - pos_i` of shape :obj:`[-1, in_channels + num_dimensions]` to shape :obj:`[-1, out_channels]`, *e.g.*, defined by :class:`torch.nn.Sequential`. (default: :obj:`None`) global_nn (torch.nn.Module, optional): A neural network :math:`\gamma_{\mathbf{\Theta}}` that maps aggregated node features of shape :obj:`[-1, out_channels]` to shape :obj:`[-1, final_out_channels]`, *e.g.*, defined by :class:`torch.nn.Sequential`. (default: :obj:`None`) add_self_loops (bool, optional): If set to :obj:`False`, will not add self-loops to the input graph. (default: :obj:`True`) **kwargs (optional): Additional arguments of :class:`torch_geometric.nn.conv.MessagePassing`. """ def __init__(self, local_nn: Optional[Callable] = None, global_nn: Optional[Callable] = None, add_self_loops: bool = True, **kwargs): super(PPFConv, self).__init__(aggr='max', **kwargs) self.local_nn = local_nn self.global_nn = global_nn self.add_self_loops = add_self_loops self.reset_parameters()
[docs] def reset_parameters(self): reset(self.local_nn) reset(self.global_nn)
[docs] def forward(self, x: Union[OptTensor, PairOptTensor], pos: Union[Tensor, PairTensor], normal: Union[Tensor, PairTensor], edge_index: Adj) -> Tensor: # yapf: disable """""" if not isinstance(x, tuple): x: PairOptTensor = (x, None) if isinstance(pos, Tensor): pos: PairTensor = (pos, pos) if isinstance(normal, Tensor): normal: PairTensor = (normal, normal) if self.add_self_loops: if isinstance(edge_index, Tensor): edge_index, _ = remove_self_loops(edge_index) edge_index, _ = add_self_loops(edge_index, num_nodes=pos[1].size(0)) elif isinstance(edge_index, SparseTensor): edge_index = set_diag(edge_index) # propagate_type: (x: PairOptTensor, pos: PairTensor, normal: PairTensor) # noqa out = self.propagate(edge_index, x=x, pos=pos, normal=normal, size=None) if self.global_nn is not None: out = self.global_nn(out) return out
def message(self, x_j: OptTensor, pos_i: Tensor, pos_j: Tensor, normal_i: Tensor, normal_j: Tensor) -> Tensor: msg = point_pair_features(pos_i, pos_j, normal_i, normal_j) if x_j is not None: msg = torch.cat([x_j, msg], dim=1) if self.local_nn is not None: msg = self.local_nn(msg) return msg def __repr__(self): return '{}(local_nn={}, global_nn={})'.format(self.__class__.__name__, self.local_nn, self.global_nn)